forked from pytorch/pytorch
-
Notifications
You must be signed in to change notification settings - Fork 0
/
__init__.py
1589 lines (1242 loc) · 49.3 KB
/
__init__.py
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
428
429
430
431
432
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
470
471
472
473
474
475
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
501
502
503
504
505
506
507
508
509
510
511
512
513
514
515
516
517
518
519
520
521
522
523
524
525
526
527
528
529
530
531
532
533
534
535
536
537
538
539
540
541
542
543
544
545
546
547
548
549
550
551
552
553
554
555
556
557
558
559
560
561
562
563
564
565
566
567
568
569
570
571
572
573
574
575
576
577
578
579
580
581
582
583
584
585
586
587
588
589
590
591
592
593
594
595
596
597
598
599
600
601
602
603
604
605
606
607
608
609
610
611
612
613
614
615
616
617
618
619
620
621
622
623
624
625
626
627
628
629
630
631
632
633
634
635
636
637
638
639
640
641
642
643
644
645
646
647
648
649
650
651
652
653
654
655
656
657
658
659
660
661
662
663
664
665
666
667
668
669
670
671
672
673
674
675
676
677
678
679
680
681
682
683
684
685
686
687
688
689
690
691
692
693
694
695
696
697
698
699
700
701
702
703
704
705
706
707
708
709
710
711
712
713
714
715
716
717
718
719
720
721
722
723
724
725
726
727
728
729
730
731
732
733
734
735
736
737
738
739
740
741
742
743
744
745
746
747
748
749
750
751
752
753
754
755
756
757
758
759
760
761
762
763
764
765
766
767
768
769
770
771
772
773
774
775
776
777
778
779
780
781
782
783
784
785
786
787
788
789
790
791
792
793
794
795
796
797
798
799
800
801
802
803
804
805
806
807
808
809
810
811
812
813
814
815
816
817
818
819
820
821
822
823
824
825
826
827
828
829
830
831
832
833
834
835
836
837
838
839
840
841
842
843
844
845
846
847
848
849
850
851
852
853
854
855
856
857
858
859
860
861
862
863
864
865
866
867
868
869
870
871
872
873
874
875
876
877
878
879
880
881
882
883
884
885
886
887
888
889
890
891
892
893
894
895
896
897
898
899
900
901
902
903
904
905
906
907
908
909
910
911
912
913
914
915
916
917
918
919
920
921
922
923
924
925
926
927
928
929
930
931
932
933
934
935
936
937
938
939
940
941
942
943
944
945
946
947
948
949
950
951
952
953
954
955
956
957
958
959
960
961
962
963
964
965
966
967
968
969
970
971
972
973
974
975
976
977
978
979
980
981
982
983
984
985
986
987
988
989
990
991
992
993
994
995
996
997
998
999
1000
from __future__ import annotations
from typing import Any, Union, Sequence, Optional, Tuple, List, Callable, Type, overload, cast
from enum import Enum
from functools import reduce, cmp_to_key
import operator
import weakref
import torch
# nvFuser imports are conditional on being compiled with CUDA
if hasattr(torch._C, "_nvfuser"):
from torch._C._nvfuser import DataType # type: ignore[import]
_torch_dtype_to_nvfuser_dtype_map = {
torch.cdouble: DataType.ComplexDouble,
torch.cfloat: DataType.ComplexFloat,
torch.double: DataType.Double,
torch.float: DataType.Float,
torch.half: DataType.Half,
torch.bfloat16: DataType.BFloat16,
torch.long: DataType.Int,
torch.int: DataType.Int32,
torch.bool: DataType.Bool,
# Python scalars
complex: DataType.ComplexDouble,
float: DataType.Double,
int: DataType.Int,
bool: DataType.Bool,
}
else:
_torch_dtype_to_nvfuser_dtype_map = {}
def getnvFuserDtype(dtype: Union[torch.dtype, NumberTypeType]):
"""
Translates from torch.dtype to nvFuser's DataType enum
"""
return _torch_dtype_to_nvfuser_dtype_map[dtype]
ShapeType = Union[torch.Size, List[int], Tuple[int, ...]]
StrideType = Union[List[int], Tuple[int, ...]]
DimsType = Union[int, List[int], Tuple[int, ...]]
DimsSequenceType = Union[List[int], Tuple[int, ...]]
# TODO: Type[torch.SymInt], Type[torch.SymFloat]
NumberTypeType = Union[Type[bool], Type[int], Type[float], Type[complex]]
# TODO: This needs a lot more type annotations
# NumberType = Union[bool, int, float, complex, torch.SymInt, torch.SymFloat]
NumberType = Union[bool, int, float, complex]
Number = (bool, int, float, complex, torch.SymInt, torch.SymFloat)
# I don't call it Integral because numbers.Integral includes bool, but IntLike
# does not
Dim = int
IntLike = (int, torch.SymInt)
FloatLike = (float, torch.SymFloat)
IntWithoutSymInt = int
FloatWithoutSymFloat = float
DeviceLikeType = Union[str, torch.device]
Tensor = torch.Tensor
torch_function_passthrough = {
torch.Tensor.dim,
torch.Tensor.ndim.__get__, # type: ignore[attr-defined]
torch.Tensor.numel,
torch.Tensor.size,
torch.Tensor.storage_offset,
torch.Tensor.stride,
torch.Tensor.dtype.__get__, # type: ignore[attr-defined]
torch.Tensor.is_sparse.__get__, # type: ignore[attr-defined]
torch.Tensor.shape.__get__, # type: ignore[attr-defined]
torch.Tensor.device.__get__, # type: ignore[attr-defined]
torch.Tensor.requires_grad.__get__, # type: ignore[attr-defined]
torch.Tensor.layout.__get__, # type: ignore[attr-defined]
# For TorchRefsMode only
torch.Tensor.__format__,
torch.Tensor.__repr__,
torch.Tensor.requires_grad.__get__, # type: ignore[attr-defined]
}
TensorLikeType = torch.Tensor
TensorLike = torch.Tensor
TensorSequenceType = Union[List[TensorLikeType], Tuple[TensorLikeType, ...]]
TensorOrNumberLikeType = Union[TensorLikeType, NumberType]
def same_shape(a: ShapeType, b: ShapeType) -> bool:
if len(a) != len(b):
return False
for x, y in zip(a, b):
if x != y:
return False
return True
# TODO: look at using torch.testing.assert_close instead with an option
# to just compare metadata
def compare_tensor_meta(a: TensorLikeType, b: TensorLikeType, check_strides=False):
"""
Checks that two tensor likes have the same shape,
dtype and device.
In the future this will validate additional metadata, like
strides.
"""
assert isinstance(a, TensorLike)
assert isinstance(b, TensorLike)
if not same_shape(a.shape, b.shape):
msg = "Shapes {0} and {1} are not equal!".format(a.shape, b.shape)
raise AssertionError(msg)
if a.dtype != b.dtype:
msg = "Dtypes {0} and {1} are not equal!".format(a.dtype, b.dtype)
raise AssertionError(msg)
if a.device != b.device:
# Handles special cuda:0 vs cuda case
# TODO: we should review why this happens and see about fixing it
if (str(a.device) == "cuda:0" or str(a.device) == "cuda") and (
str(b.device) == "cuda:0" or str(b.device) == "cuda"
):
pass
else:
msg = "Devices {0} and {1} are not equal!".format(a.device, b.device)
raise AssertionError(msg)
# Stride checking is currently disabled, see https://github.com/pytorch/pytorch/issues/78050
if check_strides:
same_strides, idx = check_significant_strides(a, b)
if not same_strides:
msg = (
"Stride mismatch! Strides are {0} and {1} (mismatched at {2})!".format(
a.stride(), b.stride(), idx
)
)
raise RuntimeError(msg)
if a.storage_offset() != b.storage_offset():
msg = (
"Storage offset mismatch! Storage offsets are {0} and {1}!".format(
a.storage_offset(), b.storage_offset()
)
)
raise RuntimeError(msg)
def check_significant_strides(
a: TensorLikeType, b: TensorLikeType
) -> Tuple[bool, Optional[int]]:
# NOTE: only on CUDA because CPU elementwise strides are incorrect in PyTorch
# See https://github.com/pytorch/pytorch/issues/77553
# Only compares strides that are "meaningful" -- strides for dimensions with length > 1
# and for tensors with more than one element
if (a.device.type == "cuda" or b.device.type == "cuda") and a.numel() > 0:
for idx in range(a.ndim):
if a.stride()[idx] != b.stride()[idx] and a.shape[idx] > 1:
return False, idx
return True, None
# This function is equivalent to compute_contiguous() from TensorImpl.cpp
def is_contiguous(a: TensorLikeType) -> bool:
"""
Tests whether a tensor is contiguous or not.
Tensors are contiguous when they have no elements,
one element, or when they have "nested" strides.
"""
if a.numel() < 2:
return True
expected_stride = 1
for x, y in reversed(tuple(zip(a.shape, a.stride()))):
# Skips checking strides when a dimension has length 1
if x == 1:
continue
if y != expected_stride:
return False
expected_stride = expected_stride * x
return True
# This function is equivalent to compute_channels_last_contiguous_2d() in TensorImpl.cpp
def is_channels_last_contiguous_2d(a: Tensor) -> bool:
# NHWC or not channels last 2D contiguous
if a.ndim != 4:
return False
expected_stride = 1
for idx in (1, 3, 2, 0):
length = a.shape[idx]
if length == 1:
continue
stride = a.stride()[idx]
if stride != expected_stride:
return False
expected_stride *= length
return True
def is_channels_last_contiguous_3d(a: Tensor) -> bool:
# NDHWC or not channels last 3D contiguous
if a.ndim != 5:
return False
expected_stride = 1
for idx in (1, 4, 3, 2, 0):
length = a.shape[idx]
if length == 1:
continue
stride = a.stride()[idx]
if stride != expected_stride:
return False
expected_stride *= length
return True
_memory_formats = set(
(
torch.contiguous_format,
torch.preserve_format,
torch.channels_last,
torch.channels_last_3d,
)
)
def validate_memory_format(memory_format: torch.memory_format):
check(
memory_format in _memory_formats,
lambda: f"Received unknown memory format {memory_format}!",
)
def is_contiguous_for_memory_format( # type: ignore[return]
a: Tensor, *, memory_format: torch.memory_format
) -> bool:
validate_memory_format(memory_format)
if memory_format == torch.contiguous_format:
return is_contiguous(a)
if memory_format == torch.channels_last:
return is_channels_last_contiguous_2d(a)
if memory_format == torch.channels_last_3d:
return is_channels_last_contiguous_3d(a)
check(
False,
lambda: f"is_contiguous received unsupported memory format {memory_format}",
)
# NOTE: that tensors with no elements and channels last is ???
def is_channels_last_contiguous(a: Tensor) -> bool:
"""
True when a tensor is channels-last contiguous.
This requires that:
- the tensor is conceptually either 4 (NHWC) or 5 (NDHWC) dimensions
- if we name the tensor's dimensions NCHW or NCDHW, then the strides are such that the
stride of the 'C' dimension (Cs) is 1 and the strides corresponding to
each dimension (Xs) can be ordered Cs <= Ws <= Hs <= (Ds) <= Ns and are
"nested" -- so Ws = Cs * Cl, where Cl is the length of the 'C' dimension,
for example.
"""
return is_channels_last_contiguous_2d(a) or is_channels_last_contiguous_3d(a)
def is_non_overlapping_and_dense(a: Tensor) -> bool:
"""
True when a tensor is non-overlapping and dense.
A tensor is non-overlapping and dense when there exists a permutation of
its dimensions that is contiguous.
"""
# Short-circuits if the tensor is already contiguous or channels-last contiguous
if is_contiguous(a) or is_channels_last_contiguous(a):
return True
# The following is equivalent to compute_non_overlapping_and_dense in TensorImpl.cpp
# Short-circuits for tensors of rank one, which are
# non-overlapping and "dense" if their stride is one
if a.ndim == 1:
return a.stride()[0] == 1
# Checks that there exists a permutation of the strides s.t. the tensor would be contiguous
# Sorts (length, stride) pairs by stride
lengths_and_strides = sorted(
tuple(zip(a.shape, a.stride())), key=operator.itemgetter(1)
)
expected_stride = 1
for length, stride in lengths_and_strides:
if length == 1:
continue
if stride != expected_stride:
return False
expected_stride *= length
return True
# NOTE: Based on the implementation in TensorIterator.cpp, but note that
# the note [Computing output strides] is incorrect, because it
# says that strides will be preserved even if they are not
# "non overlapping and dense", but this is incorrect. The
# output of elementwise operations are always given
# non overlapping and dense strides.
# This is also INCORRECT because it does not model TensorIterator's
# short-circuit, which can cause different strides.
def compute_elementwise_output_strides(*tensors) -> Tuple[int, ...]:
"""
Computes the output strides for elementwise operations.
"""
if len(tensors) == 0:
msg = "Can't compute elementwise output strides for zero tensors!"
raise ValueError(msg)
check_same_shape(*tensors, allow_cpu_scalar_tensors=True)
# Filters the tensors to actual tensors
tensors = tuple(
a for a in tensors if isinstance(a, TensorLike) and not is_cpu_scalar_tensor(a)
)
# Short-circuits for CPU scalar case
if len(tensors) == 0:
return ()
# Short-circuits for shapes with zero or one dimensions
# TODO: are these necessary?
ndim = tensors[0].ndim
if ndim == 0:
return ()
if ndim == 1:
return (1,)
shape = tensors[0].shape
def should_swap(idx_a, idx_b):
for tensor in tensors:
stride_a = tensor.stride()[idx_a]
stride_b = tensor.stride()[idx_b]
if stride_a == 0 or stride_b == 0:
continue
if stride_a < stride_b:
return -1
if stride_a > stride_b:
return 1
# stride_a == stride_b
if shape[idx_a] > shape[idx_b]:
return 1
# Note: this case is hit if all strides are zero,
# or all strides are equal and all dimensions have the same length
return 0
perm = list(reversed(range(ndim)))
# insertion sort with support for ambiguous comparisons
for i in range(1, ndim):
dim1 = i
for dim0 in reversed(range(i)):
comparison = should_swap(perm[dim0], perm[dim1])
if comparison > 0:
perm[dim0], perm[dim1] = perm[dim1], perm[dim0]
dim1 = dim0
elif comparison < 0:
break
permuted_shape = [-1] * ndim
for idx, x in enumerate(reversed(perm)):
permuted_shape[idx] = shape[x]
new_strides = make_contiguous_strides_for(permuted_shape)
permuted_strides = [-1] * ndim
for idx, x in enumerate(reversed(perm)):
permuted_strides[x] = new_strides[idx]
return tuple(permuted_strides)
#
# Common helper functions
#
def validate_dim_length(length: int):
"""
Validates that an object represents a valid
dimension length.
"""
assert length >= 0
def validate_shape(shape: ShapeType):
"""
Validates that a sequence represents a valid shape.
"""
assert isinstance(shape, Sequence)
for l in shape:
validate_dim_length(l)
def validate_strides(strides: StrideType):
"""
Verifies the object specifies valid strides.
"""
assert isinstance(strides, Sequence)
for stride in strides:
assert stride >= 0
def validate_idx(rank: int, idx: int):
"""
Validates that idx is a valid index for the given shape.
Assumes the index is already canonicalized.
"""
assert isinstance(idx, Dim)
assert isinstance(rank, Dim)
assert idx >= 0 and idx < rank or idx == 0
def validate_dimension_indices(rank: int, indices: DimsSequenceType):
for idx in indices:
validate_idx(rank, idx)
def validate_exclusive_idx(rank: int, ex_idx: int):
"""
Validates that ex_idx is a valid exclusive index
for the given shape.
"""
assert isinstance(ex_idx, Dim)
assert isinstance(rank, Dim)
assert ex_idx > 0 and ex_idx <= rank
# "Wraps" a dim (up to one time) for the given rank, allowing dims to be
# specified using negative indices. For scalar tensors with rank 0, then idx
# must be in the range [-1, 0]. Otherwise, idx should be in the range [-rank, rank-1].
def canonicalize_dim(rank: int, idx: int, wrap_scalar: bool = True) -> int:
if rank < 0:
msg = f"Rank cannot be negative but got {rank}"
raise IndexError(msg)
if rank == 0:
if not wrap_scalar:
msg = f"Dimension specified as {idx} but tensor has no dimensions"
raise IndexError(msg)
rank = 1
if idx >= 0 and idx < rank:
return idx
if idx < 0:
_idx = idx + rank
else:
_idx = idx
if _idx < 0 or _idx >= rank:
# Same error message as in aten/src/ATen/WrapDimUtils.h:49
msg = "Dimension out of range (expected to be in range of [{0}, {1}], but got {2})".format(
-rank, rank - 1, idx
)
raise IndexError(msg)
return _idx
# Takes a dimension or sequence of dimensions and "wraps" them,
# mapping negative offsets to positive ones
@overload
def canonicalize_dims(rank: int, indices: Sequence[int]) -> Tuple[int, ...]:
pass
@overload
def canonicalize_dims(rank: int, indices: int) -> int:
pass
def canonicalize_dims(rank, indices):
if isinstance(indices, Dim):
return canonicalize_dim(rank, indices)
return tuple(canonicalize_dim(rank, x) for x in indices)
def is_valid_permutation(rank: int, perm: DimsSequenceType) -> bool:
"""
Validates that perm is a permutation of length rank.
"""
if not isinstance(perm, Sequence):
return False
if not (tuple(sorted(perm)) == tuple(range(0, rank))):
return False
return True
def is_same_shape(a: Sequence, b: Sequence) -> bool:
"""
Compares two shapes a and b, returning True if they are the same
(their ranks and corresponding lengths match) and False otherwise.
"""
return tuple(a) == tuple(b)
def is_cpu_scalar_tensor(a: Any) -> bool:
return isinstance(a, TensorLike) and a.ndim == 0 and a.device.type == "cpu"
def check_same_device(*args, allow_cpu_scalar_tensors):
"""
Checks that all Tensors in args have the same device.
Raises a RuntimeError when:
- args contains an object whose type is not Tensor or Number
- two Tensor objects in args have different devices, unless one is a CPU scalar tensor and allow_cpu_scalar_tensors is True
"""
# Short-circuits if all (one or fewer) arguments are trivially on the same device
if len(args) <= 1:
return
# Note: cannot initialize device to the first arg's device (it may not have one)
device = None
for arg in args:
if isinstance(arg, Number):
continue
elif isinstance(arg, TensorLike):
if allow_cpu_scalar_tensors and is_cpu_scalar_tensor(arg):
continue
if device is None:
device = arg.device
if device != arg.device:
msg = (
"Tensor on device "
+ str(arg.device)
+ " is not on the expected device "
+ str(device)
+ "!"
)
raise RuntimeError(msg)
else:
msg = (
"Unexpected type when checking for same device, " + str(type(arg)) + "!"
)
raise RuntimeError(msg)
def canonicalize_device(device: DeviceLikeType) -> torch.device:
if isinstance(device, torch.device):
return device
assert isinstance(device, str)
return torch.device(device)
# Asserts if any of the following are true:
# - a non-scalar or non-Tensor is given
# - the shape of any tensors is distinct
def check_same_shape(*args, allow_cpu_scalar_tensors: bool):
"""
Checks that all Tensors in args have the same shape.
Raises a RuntimeError when:
- args contains an object whose type is not Tensor or Number
- two Tensor objects in args have different devices
"""
shape = None
for arg in args:
if isinstance(arg, Number):
continue
elif isinstance(arg, TensorLike):
if allow_cpu_scalar_tensors and is_cpu_scalar_tensor(arg):
continue
if shape is None:
shape = arg.shape
if not is_same_shape(shape, arg.shape):
msg = "Shape {0} is not the expected shape {1}!".format(
arg.shape, shape
)
raise RuntimeError(msg)
else:
msg = (
"Unexpected type when checking for same shape, " + str(type(arg)) + "!"
)
raise RuntimeError(msg)
# Acquires a common shape, if it exists, from one or more tensor arguments,
# filtering number arguments
def extract_shape(*args, allow_cpu_scalar_tensors: bool) -> Optional[ShapeType]:
shape = None
scalar_shape = None
for arg in args:
if isinstance(arg, Number):
continue
elif isinstance(arg, TensorLike):
if allow_cpu_scalar_tensors and is_cpu_scalar_tensor(arg):
scalar_shape = arg.shape
continue
if shape is None:
shape = arg.shape
if not is_same_shape(shape, arg.shape):
return None
else:
return None
return shape if shape is not None else scalar_shape
# Extracts dimensions that might be passed either as a list/tuple or as varargs.
# A typical case is Tensor.permute .
def extract_dims_from_varargs(dims: Union[DimsSequenceType, Tuple[DimsSequenceType, ...]]) -> DimsSequenceType:
if dims and isinstance(dims[0], Sequence):
assert len(dims) == 1
dims = cast(Tuple[DimsSequenceType], dims)
return dims[0]
else:
return cast(DimsSequenceType, dims)
def extract_shape_from_varargs(
shape: Union[ShapeType, Tuple[ShapeType]],
validate=True,
) -> Tuple[int, ...]:
"""
Returns a shape from varargs.
In PyTorch, operations that accept shapes often accept them as varargs, like
foo(*shape). However a user can pass the shape as a sequence of integers,
like this:
foo(1, 2, 3)
or as a sequence of integers
foo((1, 2, 3))
In the first case shape will be a tuple of integers, and in the second case it's a tuple
containing a tuple of integers. This validates those inputs and canonicalizes them
to a tuple of integers.
"""
# Handles tuple unwrapping
if len(shape) == 1 and isinstance(shape[0], Sequence):
shape = shape[0]
if validate:
validate_shape(shape) # type: ignore[arg-type]
return shape # type: ignore[return-value]
def infer_size(shape: ShapeType, numel: int) -> Tuple[int, ...]:
"""
Infers the size of a dim with size -1, if it exists.
Also checks that new shape is compatible with the number of elements.
"""
dim = None
newsize = 1
for i, d in enumerate(shape):
if d == -1:
check(dim is None, lambda: "only one dimension can be inferred")
dim = i
elif d >= 0:
newsize *= d
else:
check(False, lambda: f"invalid shape dimension {d}")
check(
numel == newsize or (dim is not None and newsize > 0 and numel % newsize == 0),
lambda: f"shape '{list(shape)}' is invalid for input of size {numel}",
)
if dim is not None:
check(
newsize != 0,
lambda: f"cannot reshape tensor fo 0 elements into shape {shape} because the "
f"unspecified dimension size -1 can be any value and is ambiguous",
)
shape = list(shape)
shape[dim] = numel // newsize
return tuple(shape)
_integer_dtypes = (torch.uint8, torch.int8, torch.int16, torch.int32, torch.int64)
_low_precision_dtypes = (torch.float16, torch.bfloat16, torch.complex32)
_float_dtypes = (torch.float16, torch.bfloat16, torch.float32, torch.float64)
_complex_dtypes = (torch.complex32, torch.complex64, torch.complex128)
def is_boolean_dtype(dtype: torch.dtype) -> bool:
assert isinstance(dtype, torch.dtype)
return dtype is torch.bool
def is_integer_dtype(dtype: torch.dtype) -> bool:
assert isinstance(dtype, torch.dtype)
return dtype in _integer_dtypes
def is_low_precision_dtype(dtype: torch.dtype) -> bool:
assert isinstance(dtype, torch.dtype)
return dtype in _low_precision_dtypes
def is_float_dtype(dtype: torch.dtype) -> bool:
assert isinstance(dtype, torch.dtype)
return dtype in _float_dtypes
def is_complex_dtype(dtype: torch.dtype) -> bool:
assert isinstance(dtype, torch.dtype)
return dtype in _complex_dtypes
def is_grad_dtype(dtype: torch.dtype) -> bool:
"""
Checks if the dtype can require a gradient.
"""
return is_float_dtype(dtype) or is_complex_dtype(dtype)
_complex_to_real_dtype_map = {
torch.complex128: torch.float64,
torch.complex64: torch.float32,
torch.complex32: torch.float16,
}
_real_to_complex_dtype_map = {
torch.float16: torch.complex32,
torch.bfloat16: torch.complex64,
torch.float32: torch.complex64,
torch.float64: torch.complex128,
}
def corresponding_real_dtype(dtype: torch.dtype) -> torch.dtype:
return _complex_to_real_dtype_map[dtype]
def corresponding_complex_dtype(dtype: torch.dtype) -> torch.dtype:
return _real_to_complex_dtype_map[dtype]
def dtype_to_type(dtype: torch.dtype) -> type:
"""
Computes the corresponding Python type (AKA "type kind") for the
given dtype.
"""
assert isinstance(dtype, torch.dtype)
if dtype is torch.bool:
return bool
if dtype in _integer_dtypes:
return int
if dtype in _float_dtypes:
return float
if dtype in _complex_dtypes:
return complex
raise ValueError("Invalid dtype!")
def dtype_to_type_ctor(dtype: torch.dtype) -> Callable[[NumberType], NumberType]:
"""
Computes the corresponding Python type constructor for the
given dtype.
"""
from torch.fx.experimental.symbolic_shapes import sym_float, sym_int
assert isinstance(dtype, torch.dtype)
if dtype is torch.bool:
return lambda x: bool(x)
if dtype in _integer_dtypes:
return sym_int
if dtype in _float_dtypes:
return sym_float
if dtype in _complex_dtypes:
# TODO: type error here is real, replace with sym_complex
return lambda x: complex(x) # type: ignore[arg-type]
raise ValueError("Invalid dtype!")
def type_to_dtype(typ: type) -> torch.dtype:
"""
Computes the corresponding dtype for a Number type.
"""
assert isinstance(typ, type)
if typ is bool:
return torch.bool
if typ in [int, torch.SymInt]:
return torch.long
if typ in [float, torch.SymFloat]:
return torch.get_default_dtype()
# TODO: sym_complex_float?
if typ is complex:
return corresponding_complex_dtype(torch.get_default_dtype())
raise ValueError("Invalid type!")
def get_dtype(x: Union[torch.Tensor, NumberType]):
if isinstance(x, torch.Tensor):
return x.dtype
else:
return type_to_dtype(type(x))
_ordered_types = (bool, int, float, complex)
def check_fp_or_complex(
dtype: torch.dtype, fn_name: str, allow_low_precision_dtypes: bool = True
):
"""
Checks whether the input is floating point or complex.
If allow_low_precision_dtypes is True, it allows having float16, bfloat16, and complex32
"""
check(
is_float_dtype(dtype) or is_complex_dtype(dtype),
lambda: f"{fn_name}: Expected a floating point or complex tensor as input. Got {dtype}",
)
check(
allow_low_precision_dtypes or not is_low_precision_dtype(dtype),
lambda: f"{fn_name}: Half precision dtypes not supported. Got {dtype}",
)
def check_is_matrix(A: TensorLikeType, f_name: str, arg_name: str = "A"):
check(
len(A.shape) >= 2,
lambda: f"{f_name}: The input tensor {arg_name} must have at least 2 dimensions.",
)
def get_higher_type(a: type, b: type) -> type:
"""
Returns the higher of the two given Number types.
The types are ordered bool -> int -> float -> complex.
"""
# Type checking
assert a in _ordered_types
assert b in _ordered_types
if a is b:
return a
for typ in _ordered_types:
if a is typ:
return b
if b is typ:
return a
raise ValueError("Unknown Python scalar type!")
# Returns the higher of two torch datatypes a and b or, if the two
# are not ordered relative to each other, the next
# higher datatype
def get_higher_dtype(
a: Optional[Union[torch.dtype, TensorLikeType, NumberType]],
b: Optional[Union[torch.dtype, TensorLikeType, NumberType]],
) -> Optional[torch.dtype]:
"""
Computes the "lowest" datatype that is weakly
"higher" than both a and b.
"""
# Type checking
assert a is None or isinstance(a, (torch.dtype, TensorLike, Number))
assert b is None or isinstance(b, (torch.dtype, TensorLike, Number))
def _extract_dtype(
x: Optional[Union[torch.dtype, TensorLikeType, NumberType]]
) -> Optional[torch.dtype]:
if x is None:
return None
if isinstance(x, torch.dtype):
return x
if isinstance(x, TensorLike):
return x.dtype
if isinstance(x, Number):
return type_to_dtype(type(x))
raise RuntimeError("Unexpected type given to _extract_dtype!")
a, b = _extract_dtype(a), _extract_dtype(b)
if a is b:
return a
if a is None:
return b
if b is None:
return a
ordered_datatypes = (
(torch.bool,),
(torch.uint8, torch.int8),
(torch.int16,),
(torch.int32,),
(torch.int64,),
(torch.float16, torch.bfloat16),
(torch.float32,),
(torch.float64,),
(torch.complex32,),
(torch.complex64,),
(torch.complex128,),
)
for idx, dtypes in enumerate(ordered_datatypes):
if a in dtypes and b in dtypes:
return ordered_datatypes[idx + 1][0]
if a in dtypes:
return b
if b in dtypes:
return a
raise RuntimeError("Unexpected termination!")
def check_pin_memory(pin_memory: bool):
check(not pin_memory, lambda: "PrimTorch does not support pinned memory", NotImplementedError)
def check_layout(layout: torch.layout):
check(layout == torch.strided, lambda: f"PrimTorch doesn't support layout={layout}", NotImplementedError)
# TODO: maybe unify with can_cast_to?
def is_weakly_lesser_type(a: type, b: type) -> bool:
"""
Compares two types, a and b, returning True if a is weakly "less" than b.
The comparison is determined by the following type ordering: bool, int, float, complex.
"""
ordered_types = (
bool,
int,
float,
complex,
)
assert a in ordered_types
assert b in ordered_types
for typ in ordered_types:
if a == typ:
return True