Add new np interfaces and add graph support

pull/11759/head
yanglf1121 4 years ago
parent 90777cd3bf
commit c5ea8223f5

@ -22,36 +22,59 @@ Note:
- array_ops.py defines all the array operation interfaces.
- array_creations.py defines all the array generation interfaces.
- math_ops.py defines all the math operations on tensors.
- logic_ops.py defines all the logical operations on tensors.
- dtypes.py defines all the mindspore.numpy dtypes (mainly redirected from mindspore)
"""
from .array_ops import (transpose, expand_dims, squeeze, rollaxis, swapaxes, reshape,
ravel, concatenate, where, atleast_1d, atleast_2d, atleast_3d,
column_stack, hstack, dstack, vstack, stack, unique)
column_stack, hstack, dstack, vstack, stack, unique, moveaxis,
tile, broadcast_to, broadcast_arrays, roll, append, split, vsplit,
flip, flipud, fliplr, hsplit, dsplit, take_along_axis, take, repeat)
from .array_creations import copy_ as copy
from .array_creations import (array, asarray, asfarray, ones, zeros, full, arange,
linspace, logspace, eye, identity, empty, empty_like,
ones_like, zeros_like, full_like, diagonal, tril, triu,
tri, trace)
tri, trace, cumsum, meshgrid, mgrid, ogrid, diagflat,
diag, diag_indices, ix_)
from .dtypes import (int_, int8, int16, int32, int64, uint, uint8, uint16,
uint32, uint64, float_, float16, float32, float64, bool_, inf,
numeric_types)
from .math_ops import (mean, inner, add, subtract, multiply, divide, power,
dot, outer, tensordot, absolute)
uint32, uint64, float_, float16, float32, float64, bool_, inf, nan,
numeric_types, PINF, NINF)
from .math_ops import (mean, inner, add, subtract, multiply, divide, true_divide, power,
dot, outer, tensordot, absolute, std, var, average, minimum,
matmul, square, sqrt, reciprocal, log, maximum, heaviside, amax, amin,
hypot, float_power, floor, ptp, deg2rad, rad2deg, count_nonzero,
positive, negative, clip, floor_divide, remainder, fix, fmod, trunc,
exp, expm1)
from .logic_ops import (not_equal, less_equal, less, greater_equal, greater, equal, isfinite,
isnan, isinf, isposinf, isneginf, isscalar)
mod = remainder
fabs = absolute
array_ops_module = ['transpose', 'expand_dims', 'squeeze', 'rollaxis', 'swapaxes', 'reshape',
'ravel', 'concatenate', 'where', 'atleast_1d', 'atleast_2d', 'atleast_3d',
'column_stack', 'hstack', 'dstack', 'vstack', 'stack', 'unique']
'column_stack', 'hstack', 'dstack', 'vstack', 'stack', 'unique', 'moveaxis',
'tile', 'broadcast_to', 'broadcast_arrays', 'append', 'roll', 'split', 'vsplit',
'flip', 'flipud', 'fliplr', 'hsplit', 'dsplit', 'take_along_axis', 'take',
'repeat']
array_creations_module = ['array', 'asarray', 'asfarray', 'ones', 'zeros', 'full', 'arange',
'linspace', 'logspace', 'eye', 'identity', 'empty', 'empty_like',
'ones_like', 'zeros_like', 'full_like', 'diagonal', 'tril', 'triu',
'tri', 'trace']
'tri', 'trace', 'meshgrid', 'mgrid', 'ogrid', 'diagflat', 'diag',
'diag_indices', 'ix_', 'cumsum']
math_module = ['mean', 'inner', 'add', 'subtract', 'multiply', 'divide', 'power',
'dot', 'outer', 'tensordot', 'absolute']
math_module = ['mean', 'inner', 'add', 'subtract', 'multiply', 'divide', 'true_divide', 'power',
'dot', 'outer', 'tensordot', 'absolute', 'std', 'var', 'average', 'not_equal',
'minimum', 'matmul', 'square', 'sqrt', 'reciprocal', 'log', 'maximum',
'heaviside', 'amax', 'amin', 'hypot', 'float_power', 'floor', 'ptp', 'deg2rad',
'rad2deg', 'count_nonzero', 'positive', 'negative', 'clip', 'floor_divide',
'remainder', 'mod', 'fix', 'fmod', 'trunc', 'exp', 'expm1', 'fabs']
__all__ = array_ops_module + array_creations_module + math_module + numeric_types
logic_module = ['not_equal', 'less_equal', 'less', 'greater_equal', 'greater', 'equal', 'isfinite',
'isnan', 'isinf', 'isposinf', 'isneginf', 'isscalar']
__all__ = array_ops_module + array_creations_module + math_module + logic_module + numeric_types
__all__.sort()

File diff suppressed because it is too large Load Diff

File diff suppressed because it is too large Load Diff

@ -22,7 +22,12 @@ from ..common.dtype import (int8, int16, int32, int64, uint8, uint16, uint32, ui
# backend for now.
inf = float('inf')
PINF = float('inf')
NINF = float('-inf')
nan = float('nan')
# all three of inf, PINF, and NINF are defined in the original numpy, and as we aim for
# consistency same thing is done here
pi = 3.141592653589793
int_ = int32
uint = uint32

File diff suppressed because it is too large Load Diff

File diff suppressed because it is too large Load Diff

@ -16,11 +16,11 @@
import numpy as onp
import mindspore.context as context
from ..common import Tensor
from ..ops import functional as F
from ..common import dtype as mstype
from .utils_const import _tile_size
from .utils_const import _tile_size, _add_unit_axes, _raise_type_error
def _deep_list(array_like):
@ -56,10 +56,9 @@ def _deep_tensor_to_nparray(array_like):
def _check_input_for_asarray(array_like):
"""check whether array_like argument is a valid type for np.asarray conversion"""
if isinstance(array_like, (Tensor, list, tuple, int, float, bool, onp.ndarray)):
return True
raise TypeError("input data must be `int`, `float`, `bool`, `Tensor`, `list`, `tuple`" + \
f"or numpy.ndarray, but got {type(array_like)}")
if not isinstance(array_like, (Tensor, list, tuple, int, float, bool, onp.ndarray)):
_raise_type_error("input data must be `int`, `float`, `bool`, `Tensor`, `list`, `tuple`" + \
"or numpy.ndarray, but got ", array_like)
def _is_scalar(shape):
@ -67,16 +66,6 @@ def _is_scalar(shape):
return F.shape_mul(shape) == 1
def _is_empty(shape):
"""Checks if the shape is empty"""
return F.shape_mul(shape) == 0
def _get_device():
"""Get the current device (`GPU`, `CPU`, `Ascend`)"""
return context.get_context('device_target')
def _convert_list_tensor_to_tuple_tensor(list_of_tensor):
"""Convert a list of tensor to a tuple of tensor"""
if isinstance(list_of_tensor, list):
@ -87,19 +76,66 @@ def _convert_list_tensor_to_tuple_tensor(list_of_tensor):
return list_of_tensor
def _get_mode():
"""Get the current mode (0 is Graph mode, 1 is PyNative mode)"""
return context.get_context('mode')
def _expand(x, ndim, axis=0):
"""Expand x to ndim."""
while F.rank(x) < ndim:
x = F.expand_dims(x, axis)
return x
"""Expand x to ndim from axis, which can be 0 or -1."""
shape = _add_unit_axes(F.shape(x), ndim, axis == -1)
return F.reshape(x, shape)
def _broadcast_to(x, shape_cur, shape_to, ndim_to):
"""Broadcasts x from shape_cur to shape_to."""
size = _tile_size(shape_cur, shape_to, ndim_to)
return F.tile(x, size)
def _broadcast_to_shape(x, shape):
"""Broadcasts x from current shape to shape"""
ndim_to = len(shape)
x = _expand(x, ndim_to)
return _broadcast_to(x, F.shape(x), shape, ndim_to)
def _get_size(x, axis=None):
"""Get the number of elements along the given axis of tensor x."""
if axis is None or F.tuple_len(axis) == 0:
axis = F.make_range(x.ndim)
nums = 1
for ax in axis:
nums *= x.shape[ax]
return nums
def _check_input_tensor(*tensors):
for tensor in tensors:
if not isinstance(tensor, Tensor):
_raise_type_error('expect Tensor, but got ', F.typeof(tensor))
return True
def _convert_64_to_32(tensor):
"""Convert tensor with float64/int64 types to float32/int32."""
if tensor.dtype == mstype.float64:
return tensor.astype("float32")
if tensor.dtype == mstype.int64:
return tensor.astype("int32")
return tensor
def _get_dtype_from_scalar(*input_numbers):
"""
Get the final dtype from series of input numbers, compared with F.typeof, we
return int32/float32 for python int/float instead.
"""
bool_flag = True
int_flag = True
for number in input_numbers:
if number is not None:
if not isinstance(number, bool):
bool_flag = False
if not isinstance(number, int):
int_flag = False
if bool_flag:
return mstype.bool_
if int_flag:
return mstype.int32
return mstype.float32

File diff suppressed because it is too large Load Diff

@ -59,18 +59,25 @@ tensor_div = P.RealDiv()
tensor_floordiv = P.FloorDiv()
tensor_pow = P.Pow()
tensor_mod = P.FloorMod()
tensor_exp = P.Exp()
tensor_expm1 = P.Expm1()
strided_slice = P.StridedSlice()
same_type_shape = P.SameTypeShape()
check_bprop = P.CheckBprop()
equal = P.Equal()
not_equal = P.NotEqual()
isfinite = P.IsFinite()
assign_sub = P.AssignSub()
assign_add = P.AssignAdd()
assign = P.Assign()
square = P.Square()
sqrt = P.Sqrt()
log = P.Log()
reduce_sum = P.ReduceSum()
tensor_slice = P.Slice()
maximum = P.Maximum()
minimum = P.Minimum()
floor = P.Floor()
scalar_to_array = P.ScalarToArray()
scalar_to_tensor = P.ScalarToTensor()
@ -82,6 +89,7 @@ transpose = P.Transpose()
squeeze = P.Squeeze()
scatter_nd = P.ScatterNd()
gather = P.Gather()
gather_d = P.GatherD()
gather_nd = P.GatherNd()
scatter_update = P.ScatterUpdate()
scatter_nd_update = P.ScatterNdUpdate()

File diff suppressed because it is too large Load Diff

File diff suppressed because it is too large Load Diff

File diff suppressed because it is too large Load Diff

File diff suppressed because it is too large Load Diff

@ -0,0 +1,165 @@
# Copyright 2021 Huawei Technologies Co., Ltd
#
# Licensed under the Apache License, Version 2.0 (the "License");
# you may not use this file except in compliance with the License.
# You may obtain a copy of the License at
#
# http://www.apache.org/licenses/LICENSE-2.0
#
# Unless required by applicable law or agreed to in writing, software
# distributed under the License is distributed on an "AS IS" BASIS,
# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
# See the License for the specific language governing permissions and
# limitations under the License.
# ============================================================================
"""utility functions for mindspore.numpy st tests"""
import functools
import numpy as onp
import mindspore.numpy as mnp
def match_array(actual, expected, error=0):
if isinstance(actual, int):
actual = onp.asarray(actual)
if isinstance(expected, int):
expected = onp.asarray(expected)
if error > 0:
onp.testing.assert_almost_equal(actual.tolist(), expected.tolist(),
decimal=error)
else:
onp.testing.assert_equal(actual.tolist(), expected.tolist())
def check_all_results(onp_results, mnp_results, error=0):
"""Check all results from numpy and mindspore.numpy"""
for i, _ in enumerate(onp_results):
match_array(onp_results[i], mnp_results[i].asnumpy())
def check_all_unique_results(onp_results, mnp_results):
"""
Check all results from numpy and mindspore.numpy.
Args:
onp_results (Union[tuple of numpy.arrays, numpy.array])
mnp_results (Union[tuple of Tensors, Tensor])
"""
for i, _ in enumerate(onp_results):
if isinstance(onp_results[i], tuple):
for j in range(len(onp_results[i])):
match_array(onp_results[i][j],
mnp_results[i][j].asnumpy(), error=7)
else:
match_array(onp_results[i], mnp_results[i].asnumpy(), error=7)
def run_non_kw_test(mnp_fn, onp_fn, test_case):
"""Run tests on functions with non keyword arguments"""
for i in range(len(test_case.arrs)):
arrs = test_case.arrs[:i]
match_res(mnp_fn, onp_fn, *arrs)
for i in range(len(test_case.scalars)):
arrs = test_case.scalars[:i]
match_res(mnp_fn, onp_fn, *arrs)
for i in range(len(test_case.expanded_arrs)):
arrs = test_case.expanded_arrs[:i]
match_res(mnp_fn, onp_fn, *arrs)
for i in range(len(test_case.nested_arrs)):
arrs = test_case.nested_arrs[:i]
match_res(mnp_fn, onp_fn, *arrs)
def rand_int(*shape):
"""return an random integer array with parameter shape"""
res = onp.random.randint(low=1, high=5, size=shape)
if isinstance(res, onp.ndarray):
return res.astype(onp.float32)
return float(res)
# return an random boolean array
def rand_bool(*shape):
return onp.random.rand(*shape) > 0.5
def match_res(mnp_fn, onp_fn, *arrs, **kwargs):
"""Checks results from applying mnp_fn and onp_fn on arrs respectively"""
mnp_arrs = map(functools.partial(mnp.asarray, dtype='float32'), arrs)
error = kwargs.get('error', 0)
kwargs.pop('error', None)
mnp_res = mnp_fn(*mnp_arrs, **kwargs)
onp_res = onp_fn(*arrs, **kwargs)
match_all_arrays(mnp_res, onp_res, error=error)
def match_all_arrays(mnp_res, onp_res, error=0):
if isinstance(mnp_res, (tuple, list)):
assert len(mnp_res) == len(onp_res)
for actual, expected in zip(mnp_res, onp_res):
match_array(actual.asnumpy(), expected, error)
else:
match_array(mnp_res.asnumpy(), onp_res, error)
def match_meta(actual, expected):
# float64 and int64 are not supported, and the default type for
# float and int are float32 and int32, respectively
if expected.dtype == onp.float64:
expected = expected.astype(onp.float32)
elif expected.dtype == onp.int64:
expected = expected.astype(onp.int32)
assert actual.shape == expected.shape
assert actual.dtype == expected.dtype
def run_binop_test(mnp_fn, onp_fn, test_case):
for arr in test_case.arrs:
match_res(mnp_fn, onp_fn, arr, arr)
for scalar in test_case.scalars:
match_res(mnp_fn, onp_fn, arr, scalar)
match_res(mnp_fn, onp_fn, scalar, arr)
for scalar1 in test_case.scalars:
for scalar2 in test_case.scalars:
match_res(mnp_fn, onp_fn, scalar1, scalar2)
for expanded_arr1 in test_case.expanded_arrs:
for expanded_arr2 in test_case.expanded_arrs:
match_res(mnp_fn, onp_fn, expanded_arr1, expanded_arr2)
for broadcastable1 in test_case.broadcastables:
for broadcastable2 in test_case.broadcastables:
match_res(mnp_fn, onp_fn, broadcastable1, broadcastable2)
def run_unary_test(mnp_fn, onp_fn, test_case, error=0):
for arr in test_case.arrs:
match_res(mnp_fn, onp_fn, arr, error=error)
for arr in test_case.scalars:
match_res(mnp_fn, onp_fn, arr, error=error)
for arr in test_case.expanded_arrs:
match_res(mnp_fn, onp_fn, arr, error=error)
def run_multi_test(mnp_fn, onp_fn, arrs, error=0):
mnp_arrs = map(mnp.asarray, arrs)
for actual, expected in zip(mnp_fn(*mnp_arrs), onp_fn(*arrs)):
match_array(actual.asnumpy(), expected, error)
def run_single_test(mnp_fn, onp_fn, arr, error=0):
mnp_arr = mnp.asarray(arr)
for actual, expected in zip(mnp_fn(mnp_arr), onp_fn(arr)):
if isinstance(expected, tuple):
for actual_arr, expected_arr in zip(actual, expected):
match_array(actual_arr.asnumpy(), expected_arr, error)
match_array(actual.asnumpy(), expected, error)
Loading…
Cancel
Save