burn_ndarray/
tensor.rs

1use core::mem;
2
3use burn_backend::{
4    DType, Element, QTensorPrimitive, Shape, TensorData, TensorMetadata,
5    quantization::{QParams, QuantLevel, QuantMode, QuantScheme, QuantValue},
6};
7
8use crate::NdArrayStorage;
9use crate::ops::quantization::{QuantizationStrategy, SymmetricQuantization};
10use alloc::vec::Vec;
11use ndarray::{ArcArray, ArrayD, IxDyn};
12
13/// Concrete storage type for ndarray (owned with COW semantics via Arc)
14pub type SharedArray<E> = ArcArray<E, IxDyn>;
15
16/// Tensor primitive used by the [ndarray backend](crate::NdArray).
17///
18/// Supports both owned and borrowed (zero-copy) data via [`NdArrayStorage`].
19/// When data is borrowed from external sources (like burnpack files),
20/// it remains zero-copy until a mutating operation is performed.
21#[derive(Debug, Clone)]
22#[allow(missing_docs)]
23pub enum NdArrayTensor {
24    F64(NdArrayStorage<f64>),
25    F32(NdArrayStorage<f32>),
26    I64(NdArrayStorage<i64>),
27    I32(NdArrayStorage<i32>),
28    I16(NdArrayStorage<i16>),
29    I8(NdArrayStorage<i8>),
30    U64(NdArrayStorage<u64>),
31    U32(NdArrayStorage<u32>),
32    U16(NdArrayStorage<u16>),
33    U8(NdArrayStorage<u8>),
34    Bool(NdArrayStorage<bool>),
35}
36
37impl NdArrayTensor {
38    /// Extract bool array, converting to owned if necessary.
39    pub(crate) fn bool(self) -> SharedArray<bool> {
40        match self {
41            NdArrayTensor::Bool(storage) => storage.into_shared(),
42            _ => unimplemented!("Expected bool tensor, got {:?}", self.dtype()),
43        }
44    }
45
46    /// Returns true if this tensor uses borrowed (zero-copy) storage.
47    #[inline]
48    pub fn is_borrowed(&self) -> bool {
49        macro_rules! check {
50            ($($variant:ident),*) => {
51                match self {
52                    $(NdArrayTensor::$variant(s) => s.is_borrowed(),)*
53                }
54            };
55        }
56        check!(F64, F32, I64, I32, I16, I8, U64, U32, U16, U8, Bool)
57    }
58}
59
60pub(crate) fn cast_to_dtype<E1: Element>(array: SharedArray<E1>, dtype: DType) -> NdArrayTensor
61where
62    NdArrayTensor: From<SharedArray<E1>>,
63{
64    fn cast<E1: Element, E2: Element>(array: SharedArray<E1>) -> SharedArray<E2> {
65        array.mapv(|a| a.elem()).into_shared()
66    }
67
68    if E1::dtype() == dtype {
69        return array.into();
70    }
71
72    match dtype {
73        DType::F64 => cast::<E1, f64>(array).into(),
74        DType::F32 => cast::<E1, f32>(array).into(),
75        DType::Flex32 => cast::<E1, f32>(array).into(),
76        DType::I64 => cast::<E1, i64>(array).into(),
77        DType::I32 => cast::<E1, i32>(array).into(),
78        DType::I16 => cast::<E1, i16>(array).into(),
79        DType::I8 => cast::<E1, i8>(array).into(),
80        DType::U64 => cast::<E1, u64>(array).into(),
81        DType::U32 => cast::<E1, u32>(array).into(),
82        DType::U16 => cast::<E1, u16>(array).into(),
83        DType::U8 => cast::<E1, u8>(array).into(),
84        DType::Bool => cast::<E1, bool>(array).into(),
85        dtype => panic!("Unsupported dtype: {dtype:?}"),
86    }
87}
88
89macro_rules! impl_from {
90    ($($ty: ty => $dtype: ident),*) => {
91        // From SharedArray (owned) -> NdArrayTensor
92        $(impl From<SharedArray<$ty>> for NdArrayTensor {
93           fn from(value: SharedArray<$ty>) -> NdArrayTensor {
94                NdArrayTensor::$dtype(NdArrayStorage::from_owned(value))
95           }
96        })*
97
98        // From NdArrayStorage -> NdArrayTensor
99        $(impl From<NdArrayStorage<$ty>> for NdArrayTensor {
100           fn from(value: NdArrayStorage<$ty>) -> NdArrayTensor {
101                NdArrayTensor::$dtype(value)
102           }
103        })*
104    };
105}
106
107impl_from!(
108    f64 => F64, f32 => F32,
109    i64 => I64, i32 => I32, i16 => I16, i8 => I8,
110    u64 => U64, u32 => U32, u16 => U16, u8 => U8,
111    bool => Bool
112);
113
114/// Macro to execute an operation on a given element type.
115///
116/// Extracts the storage from NdArrayTensor, converts to SharedArray, and passes to operation.
117///
118/// # Panics
119/// Since there is no automatic type cast at this time, binary operations for different
120/// floating point precision data types will panic with a data type mismatch.
121#[macro_export]
122macro_rules! execute_with_dtype {
123    (($lhs:expr, $rhs:expr),$element:ident,  $op:expr, [$($dtype: ident => $ty: ty),*]) => {{
124        let lhs_dtype = burn_backend::TensorMetadata::dtype(&$lhs);
125        let rhs_dtype = burn_backend::TensorMetadata::dtype(&$rhs);
126        match ($lhs, $rhs) {
127            $(
128                ($crate::NdArrayTensor::$dtype(lhs), $crate::NdArrayTensor::$dtype(rhs)) => {
129                    #[allow(unused)]
130                    type $element = $ty;
131                    // Convert storage to SharedArray for compatibility with existing operations
132                    $op(lhs.into_shared(), rhs.into_shared()).into()
133                }
134            )*
135            _ => panic!(
136                "Data type mismatch (lhs: {:?}, rhs: {:?})",
137                lhs_dtype, rhs_dtype
138            ),
139        }
140    }};
141    // Binary op: type automatically inferred by the compiler
142    (($lhs:expr, $rhs:expr), $op:expr) => {{
143        $crate::execute_with_dtype!(($lhs, $rhs), E, $op)
144    }};
145
146    // Binary op: generic type cannot be inferred for an operation
147    (($lhs:expr, $rhs:expr), $element:ident, $op:expr) => {{
148        $crate::execute_with_dtype!(($lhs, $rhs), $element, $op, [
149            F64 => f64, F32 => f32,
150            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
151            U64 => u64, U32 => u32, U16 => u16, U8 => u8,
152            Bool => bool
153        ])
154    }};
155
156    ($tensor:expr, $element:ident, $op:expr, [$($dtype: ident => $ty: ty),*]) => {{
157        match $tensor {
158            $(
159                $crate::NdArrayTensor::$dtype(storage) => {
160                    #[allow(unused)]
161                    type $element = $ty;
162                    // Convert to SharedArray for compatibility with most operations
163                    $op(storage.into_shared()).into()
164                }
165            )*
166            #[allow(unreachable_patterns)]
167            other => unimplemented!("unsupported dtype: {:?}", other.dtype())
168        }
169    }};
170    // Unary op: type automatically inferred by the compiler
171    ($tensor:expr, $op:expr) => {{
172        $crate::execute_with_dtype!($tensor, E, $op)
173    }};
174
175    // Unary op: generic type cannot be inferred for an operation
176    ($tensor:expr, $element:ident, $op:expr) => {{
177        $crate::execute_with_dtype!($tensor, $element, $op, [
178            F64 => f64, F32 => f32,
179            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
180            U64 => u64, U32 => u32, U16 => u16, U8 => u8,
181            Bool => bool
182        ])
183    }};
184}
185
186/// Macro to execute an operation a given element type.
187/// Only handles float types.
188///
189/// # Panics
190/// Since there is no automatic type cast at this time, binary operations for different
191/// floating point precision data types will panic with a data type mismatch.
192#[macro_export]
193macro_rules! execute_with_float_dtype {
194    // Binary op: type automatically inferred by the compiler
195    (($lhs:expr, $rhs:expr), $op:expr) => {{
196        $crate::execute_with_float_dtype!(($lhs, $rhs), E, $op)
197    }};
198
199    // Binary op: generic type cannot be inferred for an operation
200    (($lhs:expr, $rhs:expr), $element:ident, $op:expr) => {{
201        $crate::execute_with_dtype!(($lhs, $rhs), $element, $op, [
202            F64 => f64, F32 => f32
203        ])
204    }};
205
206    // Unary op: type automatically inferred by the compiler
207    ($tensor:expr, $op:expr) => {{
208        $crate::execute_with_float_dtype!($tensor, E, $op)
209    }};
210
211    // Unary op: generic type cannot be inferred for an operation
212    ($tensor:expr, $element:ident, $op:expr) => {{
213        $crate::execute_with_dtype!($tensor, $element, $op, [
214            F64 => f64, F32 => f32
215        ])
216    }};
217}
218
219/// Macro to execute an operation a given element type.
220/// Only handles int types.
221///
222/// # Panics
223/// Since there is no automatic type cast at this time, binary operations for different
224/// floating point precision data types will panic with a data type mismatch.
225#[macro_export]
226macro_rules! execute_with_int_dtype {
227    // Binary op: type automatically inferred by the compiler
228    (($lhs:expr, $rhs:expr), $op:expr) => {{
229        $crate::execute_with_int_dtype!(($lhs, $rhs), E, $op)
230    }};
231
232    // Binary op: generic type cannot be inferred for an operation
233    (($lhs:expr, $rhs:expr), $element:ident, $op:expr) => {{
234        $crate::execute_with_dtype!(($lhs, $rhs), $element, $op, [
235            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
236            U64 => u64, U32 => u32, U16 => u16, U8 => u8
237        ])
238    }};
239
240    // Unary op: type automatically inferred by the compiler
241    ($tensor:expr, $op:expr) => {{
242        $crate::execute_with_int_dtype!($tensor, E, $op)
243    }};
244
245    // Unary op: generic type cannot be inferred for an operation
246    ($tensor:expr, $element:ident, $op:expr) => {{
247        $crate::execute_with_dtype!($tensor, $element, $op, [
248            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
249            U64 => u64, U32 => u32, U16 => u16, U8 => u8
250        ])
251    }};
252}
253
254/// Macro to execute an operation a given element type.
255/// Only handles numeric types
256///
257/// # Panics
258/// Since there is no automatic type cast at this time, binary operations for different
259/// floating point precision data types will panic with a data type mismatch.
260#[macro_export]
261macro_rules! execute_with_numeric_dtype {
262    // Binary op: type automatically inferred by the compiler
263    (($lhs:expr, $rhs:expr), $op:expr) => {{
264        $crate::execute_with_numeric_dtype!(($lhs, $rhs), E, $op)
265    }};
266
267    // Binary op: generic type cannot be inferred for an operation
268    (($lhs:expr, $rhs:expr), $element:ident, $op:expr) => {{
269        $crate::execute_with_dtype!(($lhs, $rhs), $element, $op, [
270            F64 => f64, F32 => f32,
271            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
272            U64 => u64, U32 => u32, U16 => u16, U8 => u8
273        ])
274    }};
275
276    // Unary op: type automatically inferred by the compiler
277    ($tensor:expr, $op:expr) => {{
278        $crate::execute_with_numeric_dtype!($tensor, E, $op)
279    }};
280
281    // Unary op: generic type cannot be inferred for an operation
282    ($tensor:expr, $element:ident, $op:expr) => {{
283        $crate::execute_with_dtype!($tensor, $element, $op, [
284            F64 => f64, F32 => f32,
285            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
286            U64 => u64, U32 => u32, U16 => u16, U8 => u8
287        ])
288    }};
289}
290
291/// Macro to execute a cat operation on a given set of element types.
292///
293/// Uses zero-copy views from storage for concatenation.
294///
295/// # Panics
296/// Since there is no automatic type cast at this time, binary operations for different
297/// floating point precision data types will panic with a data type mismatch.
298#[macro_export]
299macro_rules! cat_with_dtype {
300    ($tensors: expr, $dim: expr, [$($dtype: ident),*]) => {
301        match &$tensors[0] {
302            $(NdArrayTensor::$dtype(_) => {
303                let tensors = $tensors
304                    .iter()
305                    .map(|t| {
306                        if let NdArrayTensor::$dtype(storage) = t {
307                            // Use storage.view() for zero-copy access
308                            storage.view()
309                        } else {
310                            panic!("Concatenate data type mismatch (expected {:?}, got {:?})", $tensors[0].dtype(), t.dtype())
311                        }
312                    })
313                    .collect::<Vec<_>>();
314                NdArrayOps::concatenate(&tensors, $dim).into()
315            })*
316            _ => panic!("Unsupported dtype: {:?}", $tensors[0].dtype())
317        }
318    };
319}
320
321impl TensorMetadata for NdArrayTensor {
322    fn dtype(&self) -> DType {
323        match self {
324            NdArrayTensor::F64(_) => DType::F64,
325            NdArrayTensor::F32(_) => DType::F32,
326            NdArrayTensor::I64(_) => DType::I64,
327            NdArrayTensor::I32(_) => DType::I32,
328            NdArrayTensor::I16(_) => DType::I16,
329            NdArrayTensor::I8(_) => DType::I8,
330            NdArrayTensor::U64(_) => DType::U64,
331            NdArrayTensor::U32(_) => DType::U32,
332            NdArrayTensor::U16(_) => DType::U16,
333            NdArrayTensor::U8(_) => DType::U8,
334            NdArrayTensor::Bool(_) => DType::Bool,
335        }
336    }
337
338    fn shape(&self) -> Shape {
339        // Use storage's shape method (works for both borrowed and owned)
340        macro_rules! get_shape {
341            ($($variant:ident),*) => {
342                match self {
343                    $(NdArrayTensor::$variant(storage) => Shape::from(storage.shape().to_vec()),)*
344                }
345            };
346        }
347        get_shape!(F64, F32, I64, I32, I16, I8, U64, U32, U16, U8, Bool)
348    }
349
350    fn rank(&self) -> usize {
351        self.shape().num_dims()
352    }
353}
354
355pub(crate) trait ShapeOps {
356    fn num_dims(self) -> usize;
357    fn num_elements(self) -> usize;
358    fn dims<const N: usize>(self) -> [usize; N];
359    fn into_shape(self) -> Shape;
360}
361
362impl ShapeOps for &[usize] {
363    fn num_dims(self) -> usize {
364        self.len()
365    }
366
367    fn num_elements(self) -> usize {
368        self.iter().product()
369    }
370
371    fn dims<const N: usize>(self) -> [usize; N] {
372        self.try_into().unwrap()
373    }
374
375    fn into_shape(self) -> Shape {
376        Shape {
377            dims: self.to_vec(),
378        }
379    }
380}
381
382mod utils {
383    use burn_std::tensor::is_contiguous;
384
385    use super::*;
386
387    impl NdArrayTensor {
388        pub(crate) fn into_data(self) -> TensorData {
389            let shape = self.shape();
390            let contiguous = self.is_contiguous();
391
392            fn inner<E: Element>(
393                shape: Shape,
394                is_contiguous: bool,
395                array: ArcArray<E, IxDyn>,
396            ) -> TensorData {
397                let vec = if is_contiguous {
398                    match array.try_into_owned_nocopy() {
399                        Ok(owned) => {
400                            let (mut vec, offset) = owned.into_raw_vec_and_offset();
401                            if let Some(offset) = offset {
402                                vec.drain(..offset);
403                            }
404                            if vec.len() > shape.num_elements() {
405                                vec.drain(shape.num_elements()..vec.len());
406                            }
407                            vec
408                        }
409                        Err(array) => array.into_iter().collect(),
410                    }
411                } else {
412                    array.into_iter().collect()
413                };
414
415                TensorData::new(vec, shape)
416            }
417
418            // Convert storage to owned array before extracting data
419            execute_with_dtype!(self, |arr| inner(shape, contiguous, arr))
420        }
421
422        pub(crate) fn is_contiguous(&self) -> bool {
423            // For borrowed data, we assume it's contiguous (it came from TensorData which is contiguous)
424            // For owned data, we check the strides
425            macro_rules! check_contiguous {
426                ($($variant:ident),*) => {
427                    match self {
428                        $(NdArrayTensor::$variant(storage) => {
429                            match storage {
430                                NdArrayStorage::Borrowed { .. } => {
431                                    // Borrowed storage requires contiguous row-major data
432                                    // (see NdArrayStorage::from_borrowed documentation)
433                                    true
434                                }
435                                NdArrayStorage::Owned(array) => {
436                                    let shape = array.shape();
437                                    let mut strides = Vec::with_capacity(array.strides().len());
438                                    for &stride in array.strides() {
439                                        if stride <= 0 {
440                                            return false;
441                                        }
442                                        strides.push(stride as usize);
443                                    }
444                                    is_contiguous(shape, &strides)
445                                }
446                            }
447                        })*
448                    }
449                };
450            }
451            check_contiguous!(F64, F32, I64, I32, I16, I8, U64, U32, U16, U8, Bool)
452        }
453    }
454}
455
456/// Converts a slice of usize to a typed dimension.
457#[macro_export(local_inner_macros)]
458macro_rules! to_typed_dims {
459    (
460        $n:expr,
461        $dims:expr,
462        justdim
463    ) => {{
464        let mut dims = [0; $n];
465        for i in 0..$n {
466            dims[i] = $dims[i];
467        }
468        let dim: Dim<[usize; $n]> = Dim(dims);
469        dim
470    }};
471}
472
473/// Reshapes an array into a tensor.
474#[macro_export(local_inner_macros)]
475macro_rules! reshape {
476    (
477        ty $ty:ty,
478        n $n:expr,
479        shape $shape:expr,
480        array $array:expr
481    ) => {{
482        let dim = $crate::to_typed_dims!($n, $shape.dims, justdim);
483        let array = match $array.is_standard_layout() {
484            true => {
485                match $array.to_shape(dim) {
486                    Ok(val) => val.into_shared(),
487                    Err(err) => {
488                        core::panic!("Shape should be compatible shape={dim:?}: {err:?}");
489                    }
490                }
491            },
492            false => $array.to_shape(dim).unwrap().as_standard_layout().into_shared(),
493        };
494        array.into_dyn()
495    }};
496    (
497        ty $ty:ty,
498        shape $shape:expr,
499        array $array:expr,
500        d $D:expr
501    ) => {{
502        match $D {
503            1 => reshape!(ty $ty, n 1, shape $shape, array $array),
504            2 => reshape!(ty $ty, n 2, shape $shape, array $array),
505            3 => reshape!(ty $ty, n 3, shape $shape, array $array),
506            4 => reshape!(ty $ty, n 4, shape $shape, array $array),
507            5 => reshape!(ty $ty, n 5, shape $shape, array $array),
508            6 => reshape!(ty $ty, n 6, shape $shape, array $array),
509            _ => core::panic!("NdArray supports arrays up to 6 dimensions, received: {}", $D),
510        }
511    }};
512}
513
514impl NdArrayTensor {
515    /// Create a new [ndarray tensor](NdArrayTensor) from [data](TensorData).
516    ///
517    /// This method attempts zero-copy loading when possible. If the data has properly
518    /// aligned bytes that can be borrowed, it creates a borrowed tensor. Otherwise,
519    /// it falls back to copying the data.
520    ///
521    /// Zero-copy loading works when:
522    /// - The data's bytes are properly aligned for the element type
523    /// - The bytes can be borrowed (e.g., from mmap'd file or static data)
524    pub fn from_data(data: TensorData) -> NdArrayTensor {
525        // Try borrowed storage first, fall back to owned if not possible
526        match Self::try_from_data_borrowed(data) {
527            Ok(tensor) => tensor,
528            Err(data) => Self::from_data_owned(data),
529        }
530    }
531
532    /// Try to create a tensor with borrowed storage (zero-copy).
533    ///
534    /// Takes ownership of TensorData and returns it back on failure.
535    /// No cloning occurs - bytes are moved into storage or returned on failure.
536    ///
537    /// Returns `Err(data)` if borrowing is not possible (e.g., misaligned data).
538    fn try_from_data_borrowed(data: TensorData) -> Result<NdArrayTensor, TensorData> {
539        let TensorData {
540            bytes,
541            shape,
542            dtype,
543        } = data;
544
545        macro_rules! try_borrow {
546            ($ty:ty, $variant:ident, $bytes:expr, $shape:expr) => {
547                match NdArrayStorage::<$ty>::from_borrowed($bytes, $shape) {
548                    Ok(storage) => return Ok(NdArrayTensor::$variant(storage)),
549                    Err((bytes, shape)) => (bytes, shape),
550                }
551            };
552        }
553
554        // Try to create borrowed storage; get bytes back on failure
555        let (bytes, shape) = match dtype {
556            DType::F64 => try_borrow!(f64, F64, bytes, shape),
557            DType::F32 => try_borrow!(f32, F32, bytes, shape),
558            DType::I64 => try_borrow!(i64, I64, bytes, shape),
559            DType::I32 => try_borrow!(i32, I32, bytes, shape),
560            DType::I16 => try_borrow!(i16, I16, bytes, shape),
561            DType::I8 => try_borrow!(i8, I8, bytes, shape),
562            DType::U64 => try_borrow!(u64, U64, bytes, shape),
563            DType::U32 => try_borrow!(u32, U32, bytes, shape),
564            DType::U16 => try_borrow!(u16, U16, bytes, shape),
565            DType::U8 => try_borrow!(u8, U8, bytes, shape),
566            DType::Bool => try_borrow!(bool, Bool, bytes, shape),
567            _ => (bytes, shape), // QFloat not supported for zero-copy
568        };
569
570        Err(TensorData {
571            bytes,
572            shape,
573            dtype,
574        })
575    }
576
577    /// Create a tensor with owned storage.
578    ///
579    /// This may or may not copy data depending on whether the underlying bytes
580    /// can be reclaimed (via `try_into_vec`). If bytes are uniquely owned,
581    /// no copy occurs; otherwise data is copied to a new allocation.
582    fn from_data_owned(mut data: TensorData) -> NdArrayTensor {
583        let shape = mem::take(&mut data.shape);
584
585        macro_rules! execute {
586            ($data: expr, [$($dtype: ident => $ty: ty),*]) => {
587                match $data.dtype {
588                    $(DType::$dtype => {
589                        match data.into_vec::<$ty>() {
590                            // Safety: TensorData checks shape validity on creation
591                            Ok(vec) => unsafe { ArrayD::from_shape_vec_unchecked(shape, vec) }.into_shared(),
592                            Err(err) => panic!("Data should have the same element type as the tensor {err:?}"),
593                        }.into()
594                    },)*
595                    other => unimplemented!("Unsupported dtype {other:?}"),
596                }
597            };
598        }
599
600        execute!(data, [
601            F64 => f64, F32 => f32,
602            I64 => i64, I32 => i32, I16 => i16, I8 => i8,
603            U64 => u64, U32 => u32, U16 => u16, U8 => u8,
604            Bool => bool
605        ])
606    }
607}
608
609/// A quantized tensor for the ndarray backend.
610#[derive(Clone, Debug)]
611pub struct NdArrayQTensor {
612    /// The quantized tensor.
613    pub qtensor: NdArrayTensor,
614    /// The quantization scheme.
615    pub scheme: QuantScheme,
616    /// The quantization parameters.
617    pub qparams: Vec<QParams<f32>>,
618}
619
620impl NdArrayQTensor {
621    /// Returns the quantization strategy, including quantization parameters, for the given tensor.
622    pub fn strategy(&self) -> QuantizationStrategy {
623        match self.scheme {
624            QuantScheme {
625                level: QuantLevel::Tensor,
626                mode: QuantMode::Symmetric,
627                value:
628                    QuantValue::Q8F
629                    | QuantValue::Q8S
630                    | QuantValue::E4M3
631                    | QuantValue::E5M2
632                    | QuantValue::Q4F
633                    | QuantValue::Q4S
634                    | QuantValue::E2M1
635                    | QuantValue::Q2F
636                    | QuantValue::Q2S,
637                ..
638            } => QuantizationStrategy::PerTensorSymmetric(SymmetricQuantization::init(
639                self.qparams[0].scales,
640                self.scheme.value,
641            )),
642            QuantScheme {
643                level: QuantLevel::Block(block_size),
644                mode: QuantMode::Symmetric,
645                value:
646                    QuantValue::Q8F
647                    | QuantValue::Q8S
648                    | QuantValue::E4M3
649                    | QuantValue::E5M2
650                    | QuantValue::Q4F
651                    | QuantValue::Q4S
652                    | QuantValue::E2M1
653                    | QuantValue::Q2F
654                    | QuantValue::Q2S,
655                ..
656            } => QuantizationStrategy::PerBlockSymmetric(
657                self.qparams
658                    .iter()
659                    .map(|q| SymmetricQuantization::init(q.scales, self.scheme.value))
660                    .collect(),
661                block_size,
662            ),
663        }
664    }
665}
666
667impl QTensorPrimitive for NdArrayQTensor {
668    fn scheme(&self) -> &QuantScheme {
669        &self.scheme
670    }
671
672    fn default_scheme() -> QuantScheme {
673        QuantScheme::default().with_store(burn_backend::quantization::QuantStore::Native)
674    }
675}
676
677impl TensorMetadata for NdArrayQTensor {
678    fn dtype(&self) -> DType {
679        DType::QFloat(self.scheme)
680    }
681
682    fn shape(&self) -> Shape {
683        self.qtensor.shape()
684    }
685
686    fn rank(&self) -> usize {
687        self.shape().num_dims()
688    }
689}
690
691#[cfg(test)]
692mod tests {
693    use crate::NdArray;
694    use alloc::vec;
695
696    use super::*;
697    use burn_backend::{
698        Distribution,
699        ops::{FloatTensorOps, QTensorOps},
700        quantization::{QuantStore, QuantizationParametersPrimitive},
701    };
702    use burn_std::rand::get_seeded_rng;
703
704    #[test]
705    fn should_support_into_and_from_data_1d() {
706        let data_expected = TensorData::random::<f32, _, _>(
707            Shape::new([3]),
708            Distribution::Default,
709            &mut get_seeded_rng(),
710        );
711        let tensor = NdArrayTensor::from_data(data_expected.clone());
712
713        let data_actual = tensor.into_data();
714
715        assert_eq!(data_expected, data_actual);
716    }
717
718    #[test]
719    fn should_support_into_and_from_data_2d() {
720        let data_expected = TensorData::random::<f32, _, _>(
721            Shape::new([2, 3]),
722            Distribution::Default,
723            &mut get_seeded_rng(),
724        );
725        let tensor = NdArrayTensor::from_data(data_expected.clone());
726
727        let data_actual = tensor.into_data();
728
729        assert_eq!(data_expected, data_actual);
730    }
731
732    #[test]
733    fn should_support_into_and_from_data_3d() {
734        let data_expected = TensorData::random::<f32, _, _>(
735            Shape::new([2, 3, 4]),
736            Distribution::Default,
737            &mut get_seeded_rng(),
738        );
739        let tensor = NdArrayTensor::from_data(data_expected.clone());
740
741        let data_actual = tensor.into_data();
742
743        assert_eq!(data_expected, data_actual);
744    }
745
746    #[test]
747    fn should_support_into_and_from_data_4d() {
748        let data_expected = TensorData::random::<f32, _, _>(
749            Shape::new([2, 3, 4, 2]),
750            Distribution::Default,
751            &mut get_seeded_rng(),
752        );
753        let tensor = NdArrayTensor::from_data(data_expected.clone());
754
755        let data_actual = tensor.into_data();
756
757        assert_eq!(data_expected, data_actual);
758    }
759
760    #[test]
761    fn should_support_qtensor_strategy() {
762        type B = NdArray<f32, i64, i8>;
763        let scale: f32 = 0.009_019_608;
764        let device = Default::default();
765
766        let tensor = B::float_from_data(TensorData::from([-1.8f32, -1.0, 0.0, 0.5]), &device);
767        let scheme = QuantScheme::default()
768            .with_value(QuantValue::Q8S)
769            .with_store(QuantStore::Native);
770        let qparams = QuantizationParametersPrimitive {
771            scales: B::float_from_data(TensorData::from([scale]), &device),
772        };
773        let qtensor: NdArrayQTensor = B::quantize(tensor, &scheme, qparams);
774
775        assert_eq!(qtensor.scheme(), &scheme);
776        assert_eq!(
777            qtensor.strategy(),
778            QuantizationStrategy::PerTensorSymmetric(SymmetricQuantization::init(
779                scale,
780                QuantValue::Q8S
781            ))
782        );
783    }
784
785    // ==========================================================================
786    // Zero-copy integration tests
787    // These tests verify end-to-end zero-copy behavior through NdArrayTensor.
788    // ==========================================================================
789
790    #[test]
791    fn zero_copy_creates_borrowed_storage() {
792        // Verify that from_data creates borrowed storage when possible.
793        // Note: For native allocations, Bytes::clone() copies data internally,
794        // but the storage type (Borrowed) is preserved, which is important for
795        // the is_unique() behavior that triggers copy-on-write.
796        use burn_std::Bytes;
797
798        let data: Vec<f32> = vec![1.0, 2.0, 3.0, 4.0];
799        let bytes = Bytes::from_elems(data);
800        let tensor_data = TensorData::from_bytes(bytes, Shape::new([2, 2]), DType::F32);
801
802        let tensor = NdArrayTensor::from_data(tensor_data);
803
804        match &tensor {
805            NdArrayTensor::F32(storage) => {
806                assert!(
807                    storage.is_borrowed(),
808                    "ZERO-COPY REGRESSION: from_data should create borrowed storage \
809                     for properly aligned TensorData with Bytes"
810                );
811                assert!(
812                    !storage.is_unique(),
813                    "ZERO-COPY REGRESSION: borrowed storage must report is_unique() == false"
814                );
815            }
816            _ => panic!("Expected F32 tensor"),
817        }
818    }
819
820    #[test]
821    fn zero_copy_data_integrity() {
822        // Verify data is correctly accessible through borrowed storage
823        use burn_std::Bytes;
824
825        let data: Vec<f32> = vec![1.0, 2.0, 3.0, 4.0];
826        let bytes = Bytes::from_elems(data);
827        let tensor_data = TensorData::from_bytes(bytes, Shape::new([2, 2]), DType::F32);
828
829        let tensor = NdArrayTensor::from_data(tensor_data);
830
831        match &tensor {
832            NdArrayTensor::F32(storage) => {
833                let view = storage.view();
834                assert_eq!(view[[0, 0]], 1.0);
835                assert_eq!(view[[0, 1]], 2.0);
836                assert_eq!(view[[1, 0]], 3.0);
837                assert_eq!(view[[1, 1]], 4.0);
838            }
839            _ => panic!("Expected F32 tensor"),
840        }
841    }
842
843    #[test]
844    fn zero_copy_fallback_when_bytes_owned() {
845        // When TensorData owns bytes exclusively, it may use the copy path
846        // This is expected behavior - verify it still works correctly
847        let data = TensorData::from([1.0f32, 2.0, 3.0, 4.0]);
848        let tensor = NdArrayTensor::from_data(data.clone());
849        let result = tensor.into_data();
850
851        assert_eq!(data, result, "Data should round-trip correctly");
852    }
853}