You can not select more than 25 topics Topics must start with a chinese character,a letter or number, can include dashes ('-') and can be up to 35 characters long.

tensor.py 18 kB

5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
5 years ago
123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499
  1. # Copyright 2020 Huawei Technologies Co., Ltd
  2. #
  3. # Licensed under the Apache License, Version 2.0 (the "License");
  4. # you may not use this file except in compliance with the License.
  5. # You may obtain a copy of the License at
  6. #
  7. # http://www.apache.org/licenses/LICENSE-2.0
  8. #
  9. # Unless required by applicable law or agreed to in writing, software
  10. # distributed under the License is distributed on an "AS IS" BASIS,
  11. # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  12. # See the License for the specific language governing permissions and
  13. # limitations under the License.
  14. # ============================================================================
  15. """Tensor implementation."""
  16. import numpy as np
  17. from mindspore import log as logger
  18. from mindspore.communication.management import get_rank, get_group_size
  19. from .._c_expression import Tensor as Tensor_
  20. from .._c_expression import MetaTensor as MetaTensor_
  21. from .._checkparam import Validator as validator
  22. from . import dtype as mstype
  23. from ._register_for_tensor import tensor_operator_registry
  24. __all__ = ['Tensor', 'MetaTensor', 'RowTensor', 'SparseTensor']
  25. np_types = (np.int8, np.int16, np.int32, np.int64,
  26. np.uint8, np.uint16, np.uint32, np.uint64, np.float16,
  27. np.float32, np.float64, np.bool_)
  28. class Tensor(Tensor_):
  29. """
  30. Tensor is used for data storage.
  31. Tensor inherits tensor object in C++.
  32. Some functions are implemented in C++ and some functions are implemented in Python.
  33. Args:
  34. input_data (Tensor, float, int, bool, tuple, list, numpy.ndarray): Input data of the tensor.
  35. dtype (:class:`mindspore.dtype`): Input data should be None, bool or numeric type defined in `mindspore.dtype`.
  36. The argument is used to define the data type of the output tensor. If it is None, the data type of the
  37. output tensor will be as same as the `input_data`. Default: None.
  38. Outputs:
  39. Tensor, with the same shape as `input_data`.
  40. Examples:
  41. >>> # initialize a tensor with input data
  42. >>> t1 = Tensor(np.zeros([1, 2, 3]), mindspore.float32)
  43. >>> assert isinstance(t1, Tensor)
  44. >>> assert t1.shape == (1, 2, 3)
  45. >>> assert t1.dtype == mindspore.float32
  46. >>>
  47. >>> # initialize a tensor with a float scalar
  48. >>> t2 = Tensor(0.1)
  49. >>> assert isinstance(t2, Tensor)
  50. >>> assert t2.dtype == mindspore.float64
  51. """
  52. def __init__(self, input_data, dtype=None):
  53. # If input data is numpy number, convert it to np array
  54. if isinstance(input_data, np_types):
  55. input_data = np.array(input_data)
  56. # If input_data is tuple/list/numpy.ndarray, it's support in check_type method.
  57. validator.check_value_type('input_data', input_data, (Tensor_, np.ndarray, list, tuple, float, int, bool),
  58. 'Tensor')
  59. valid_dtypes = (np.int8, np.int16, np.int32, np.int64, np.uint8, np.uint16, np.uint32, np.uint64,
  60. np.float16, np.float32, np.float64, np.bool_)
  61. if isinstance(input_data, np.ndarray) and input_data.dtype not in valid_dtypes:
  62. raise TypeError(f"For Tensor, the input_data is a numpy array whose data type is "
  63. f"{input_data.dtype} that is not supported to initialize a Tensor.")
  64. if isinstance(input_data, (tuple, list)):
  65. if np.array(input_data).dtype not in valid_dtypes:
  66. raise TypeError(f"For Tensor, the input_data is {input_data} that contain unsupported element.")
  67. if dtype is not None:
  68. validator.check_type_name('dtype', dtype, mstype.number_type + (mstype.bool_,), "Tensor")
  69. if isinstance(input_data, np.ndarray) and (not input_data.flags['FORC']):
  70. input_data = np.ascontiguousarray(input_data)
  71. if dtype is None:
  72. Tensor_.__init__(self, input_data)
  73. else:
  74. Tensor_.__init__(self, input_data, dtype)
  75. self._virtual_flag = False
  76. def __repr__(self):
  77. return Tensor_.__repr__(self)
  78. def __add__(self, other):
  79. out = tensor_operator_registry.get('__add__')(self, other)
  80. return out
  81. def __eq__(self, other):
  82. if not isinstance(other, (int, float, Tensor)):
  83. return False
  84. # bool type is not supported for `Equal` operator in backend.
  85. if self.dtype == mstype.bool_ or (isinstance(other, Tensor) and other.dtype == mstype.bool_):
  86. if isinstance(other, Tensor):
  87. return Tensor(np.array(self.asnumpy() == other.asnumpy()))
  88. return Tensor(np.array(self.asnumpy() == other))
  89. return tensor_operator_registry.get('__eq__')(self, other)
  90. def __ne__(self, other):
  91. if not isinstance(other, (int, float, Tensor)):
  92. return True
  93. # bool type is not supported for `NotEqual` operator in backend.
  94. if self.dtype == mstype.bool_ or (isinstance(other, Tensor) and other.dtype == mstype.bool_):
  95. return Tensor(np.array(self.asnumpy() != other.asnumpy()))
  96. return tensor_operator_registry.get('__ne__')(self, other)
  97. def __hash__(self):
  98. return hash(id(self))
  99. def __mul__(self, other):
  100. out = tensor_operator_registry.get('__mul__')(self, other)
  101. return out
  102. def __neg__(self):
  103. out = tensor_operator_registry.get('__neg__')(self)
  104. return out
  105. def __bool__(self):
  106. data = self.asnumpy()
  107. if data.shape == ():
  108. return bool(data)
  109. if data.shape == (1,):
  110. return bool(data[0])
  111. raise ValueError("The truth value of an array with several elements is ambiguous.")
  112. def __pos__(self):
  113. return self
  114. def __iadd__(self, other):
  115. return self.__add__(other)
  116. def __radd__(self, other):
  117. out = tensor_operator_registry.get('__add__')(self, other)
  118. return out
  119. def __imul__(self, other):
  120. return self.__mul__(other)
  121. def __rmul__(self, other):
  122. out = tensor_operator_registry.get('__mul__')(self, other)
  123. return out
  124. def __truediv__(self, other):
  125. out = tensor_operator_registry.get('__truediv__')(self, other)
  126. return out
  127. def __rtruediv__(self, other):
  128. out = tensor_operator_registry.get('__truediv__')(other, self)
  129. return out
  130. def __sub__(self, other):
  131. out = tensor_operator_registry.get('__sub__')(self, other)
  132. return out
  133. def __isub__(self, other):
  134. return self.__sub__(other)
  135. def __rsub__(self, other):
  136. out = tensor_operator_registry.get('__sub__')(other, self)
  137. return out
  138. def __lt__(self, other):
  139. out = tensor_operator_registry.get('__lt__')(self, other)
  140. return out
  141. def __le__(self, other):
  142. out = tensor_operator_registry.get('__le__')(self, other)
  143. return out
  144. def __getitem__(self, index):
  145. out = tensor_operator_registry.get('__getitem__')(self, index)
  146. return out
  147. def __setitem__(self, index, value):
  148. out = tensor_operator_registry.get('__setitem__')(self, index, value)
  149. self.assign_value(out)
  150. return self
  151. def __gt__(self, other):
  152. out = tensor_operator_registry.get('__gt__')(self, other)
  153. return out
  154. def __ge__(self, other):
  155. out = tensor_operator_registry.get('__ge__')(self, other)
  156. return out
  157. def __len__(self):
  158. out = tensor_operator_registry.get('shape')(self)
  159. if not out:
  160. return 1
  161. return out[0]
  162. def __mod__(self, other):
  163. return tensor_operator_registry.get('__mod__')(self, other)
  164. def __imod__(self, other):
  165. return self.__mod__(other)
  166. def __rmod__(self, other):
  167. return tensor_operator_registry.get('__mod__')(other, self)
  168. def __pow__(self, other):
  169. return tensor_operator_registry.get('__pow__')(self, other)
  170. def __floordiv__(self, other):
  171. return tensor_operator_registry.get('__floordiv__')(self, other)
  172. def __ifloordiv__(self, other):
  173. return self.__floordiv__(other)
  174. def __rfloordiv__(self, other):
  175. return tensor_operator_registry.get('__floordiv__')(other, self)
  176. def __str__(self):
  177. if self.dtype == mstype.type_none:
  178. return "Unknown Tensor type!"
  179. return str(self.asnumpy())
  180. @property
  181. def shape(self):
  182. """The shape of tensor is a tuple."""
  183. return self._shape
  184. @property
  185. def dtype(self):
  186. """The dtype of tensor is a mindspore type."""
  187. return self._dtype
  188. @property
  189. def virtual_flag(self):
  190. """Mark tensor is virtual."""
  191. return self._virtual_flag
  192. @virtual_flag.setter
  193. def virtual_flag(self, value):
  194. """The setter of virtual_flag."""
  195. if not isinstance(value, bool):
  196. raise TypeError("virtual_flag must be bool.")
  197. self._virtual_flag = value
  198. @staticmethod
  199. def from_numpy(array):
  200. """Convert numpy array to Tensor without copy data."""
  201. return Tensor(Tensor_.from_numpy(array))
  202. def asnumpy(self):
  203. """Convert tensor to numpy array."""
  204. return Tensor_.asnumpy(self)
  205. def all(self, axis=(), keep_dims=False):
  206. """
  207. Check all array elements along a given axis evaluate to True.
  208. Args:
  209. axis (Union[None, int, tuple(int)): Dimensions of reduction,
  210. when axis is None or empty tuple, reduce all dimensions.
  211. Default: (), reduce all dimensions.
  212. keep_dims (bool): Whether to keep the reduced dimensions.
  213. Default : False, don't keep these reduced dimensions.
  214. Returns:
  215. Tensor, has the same data type as x.
  216. """
  217. if axis is None:
  218. axis = ()
  219. return tensor_operator_registry.get('all')(keep_dims)(self, axis)
  220. def any(self, axis=(), keep_dims=False):
  221. """
  222. Check any array element along a given axis evaluate to True.
  223. Args:
  224. axis (Union[None, int, tuple(int)): Dimensions of reduction,
  225. when axis is None or empty tuple, reduce all dimensions.
  226. Default: (), reduce all dimensions.
  227. keep_dims (bool): Whether to keep the reduced dimensions.
  228. Default : False, don't keep these reduced dimensions.
  229. Returns:
  230. Tensor, has the same data type as x.
  231. """
  232. if axis is None:
  233. axis = ()
  234. return tensor_operator_registry.get('any')(keep_dims)(self, axis)
  235. class RowTensor:
  236. """
  237. A sparse representation of a set of tensor slices at given indices.
  238. An RowTensor is typically used to represent a subset of a larger
  239. tensor dense of shape [L0, D1, .. , DN] where L0 >> D0.
  240. The values in indices are the indices in the first dimension of the slices
  241. that have been extracted from the larger tensor.
  242. The dense tensor dense represented by an RowTensor slices has
  243. `dense[slices.indices[i], :, :, :, ...] = slices.values[i, :, :, :, ...]`.
  244. RowTensor can only be used in the `Cell`'s construct method.
  245. It is not supported in pynative mode at the moment.
  246. Args:
  247. indices (Tensor): A 1-D integer Tensor of shape [D0].
  248. values (Tensor): A Tensor of any dtype of shape [D0, D1, ..., Dn].
  249. dense_shape (tuple): An integer tuple which contains the shape
  250. of the corresponding dense tensor.
  251. Returns:
  252. RowTensor, composed of `indices`, `values`, and `dense_shape`.
  253. Examples:
  254. >>> class Net(nn.Cell):
  255. >>> def __init__(self, dense_shape):
  256. >>> super(Net, self).__init__()
  257. >>> self.dense_shape = dense_shape
  258. >>> def construct(self, indices, values):
  259. >>> x = RowTensor(indices, values, self.dense_shape)
  260. >>> return x.values, x.indices, x.dense_shape
  261. >>>
  262. >>> indices = Tensor([0])
  263. >>> values = Tensor([[1, 2]], dtype=ms.float32)
  264. >>> Net((3, 2))(indices, values)
  265. """
  266. def __init__(self, indices, values, dense_shape):
  267. "Init RowTensor"
  268. self.__indices = indices
  269. self.__values = values
  270. self.__dense_shape = dense_shape
  271. @property
  272. def indices(self):
  273. return self.__indices
  274. @property
  275. def values(self):
  276. return self.__values
  277. @property
  278. def dense_shape(self):
  279. return self.__dense_shape
  280. class SparseTensor:
  281. """
  282. A sparse representation of a set of nonzero elememts from a tensor at given indices.
  283. SparseTensor can only be used in the `Cell`'s construct method.
  284. Pynative mode not supported at the moment.
  285. For a tensor dense, its SparseTensor(indices, values, dense_shape) has
  286. `dense[indices[i]] = values[i]`.
  287. Args:
  288. indices (Tensor): A 2-D integer Tensor of shape `[N, ndims]`,
  289. where N and ndims are the number of values and number of dimensions in
  290. the SparseTensor, respectively.
  291. values (Tensor): A 1-D tensor of any type and shape `[N]`, which
  292. supplies the values for each element in indices.
  293. dense_shape (tuple): A integer tuple of size `ndims`,
  294. which specifies the dense_shape of the sparse tensor.
  295. Returns:
  296. SparseTensor, composed of `indices`, `values`, and `dense_shape`.
  297. Examples:
  298. >>> class Net(nn.Cell):
  299. >>> def __init__(self, dense_shape):
  300. >>> super(Net, self).__init__()
  301. >>> self.dense_shape = dense_shape
  302. >>> def construct(self, indices, values):
  303. >>> x = SparseTensor(indices, values, self.dense_shape)
  304. >>> return x.values, x.indices, x.dense_shape
  305. >>>
  306. >>> indices = Tensor([[0, 1], [1, 2]])
  307. >>> values = Tensor([1, 2], dtype=ms.float32)
  308. >>> Net((3, 4))(indices, values)
  309. """
  310. def __init__(self, indices, values, dense_shape):
  311. "Init SparseTensor"
  312. self.__indices = indices
  313. self.__values = values
  314. self.__dense_shape = dense_shape
  315. @property
  316. def indices(self):
  317. return self.__indices
  318. @property
  319. def values(self):
  320. return self.__values
  321. @property
  322. def dense_shape(self):
  323. return self.__dense_shape
  324. class MetaTensor(MetaTensor_):
  325. """
  326. The base class of the MetaTensor.
  327. Initialization of tensor basic attributes and model weight values.
  328. Returns:
  329. Array, an array after being initialized.
  330. """
  331. def __init__(self, dtype, shape, init=None):
  332. # check param
  333. self.init = init
  334. MetaTensor_.__init__(self, dtype, shape)
  335. def to_tensor(self, slice_index=None, shape=None, opt_shard_group=None):
  336. """
  337. Get the tensor format data of this MetaTensor.
  338. Args:
  339. slice_index (int): Slice index of a parameter's slices.
  340. It is used when initialize a slice of a parameter, it guarantees that devices
  341. using the same slice can generate the same tensor.
  342. shape (list[int]): Shape of the slice, it is used when initialize a slice of the parameter.
  343. opt_shard_group(str): Optimizer shard group which is used in auto or semi auto parallel mode
  344. to get one shard of a parameter's slice.
  345. """
  346. if self.init is None:
  347. raise TypeError("to_dense must be set MetaTensor.init, init can't be None")
  348. if shape is None:
  349. shape = self.shape
  350. try:
  351. arr = np.ndarray(shape, dtype=mstype.dtype_to_nptype(self.dtype))
  352. except ValueError:
  353. msg = "Error shape={}".format(shape)
  354. logger.error(msg)
  355. raise ValueError(msg)
  356. class seed_context:
  357. '''set and restore seed'''
  358. def __init__(self, init):
  359. self.init = init
  360. from .seed import get_seed
  361. global_seed = get_seed()
  362. self._np_seed = np.random.get_state()[1][0]
  363. self.need_set_seed = ((slice_index is not None) and (global_seed is None))
  364. def __enter__(self):
  365. if self.need_set_seed:
  366. self.seed = self.init.seed
  367. np.random.seed(slice_index)
  368. self.init.seed = slice_index
  369. def __exit__(self, ptype, value, trace):
  370. if self.need_set_seed:
  371. np.random.seed(self._np_seed)
  372. self.init.seed = self.seed
  373. with seed_context(self.init):
  374. self.init(arr)
  375. data = np.array(arr)
  376. if opt_shard_group:
  377. rank = get_rank(opt_shard_group)
  378. size = get_group_size(opt_shard_group)
  379. data = np.split(data, size)[rank]
  380. return Tensor(data, dtype=self.dtype)
  381. def _vm_compare(*args):
  382. """Implement `vm_compare` for tensor."""
  383. obj_str = args[-1]
  384. if obj_str == "shape":
  385. fn = getattr(args[0].asnumpy(), obj_str)
  386. return fn
  387. if len(args) == 2:
  388. fn = getattr(args[0].asnumpy(), obj_str)
  389. return Tensor(fn())
  390. if isinstance(args[0], Tensor):
  391. fn = getattr(args[0].asnumpy(), obj_str)
  392. y = args[1].asnumpy() if isinstance(args[1], Tensor) else args[1]
  393. else:
  394. obj_str = "__r" + obj_str[2:]
  395. fn = getattr(args[1].asnumpy(), obj_str)
  396. y = args[0]
  397. return Tensor(np.array(fn(y)))
  398. tensor_operator_registry.register('vm_compare', _vm_compare)