2018-11-17 18:03:33 -08:00
|
|
|
# Copyright 2018 Google LLC
|
|
|
|
#
|
|
|
|
# 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
|
|
|
|
#
|
|
|
|
# https://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.
|
|
|
|
|
2019-12-19 11:19:58 -08:00
|
|
|
from contextlib import contextmanager
|
2018-11-17 18:03:33 -08:00
|
|
|
import functools
|
2018-11-21 13:20:44 -08:00
|
|
|
import re
|
2018-12-13 08:56:40 -08:00
|
|
|
import os
|
2020-09-14 02:47:28 -07:00
|
|
|
import textwrap
|
2020-03-19 08:54:37 +01:00
|
|
|
from typing import Dict, Sequence, Union
|
2020-04-12 15:35:35 -04:00
|
|
|
import unittest
|
|
|
|
import warnings
|
2020-05-04 23:00:20 -04:00
|
|
|
import zlib
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-06-24 16:24:33 -07:00
|
|
|
from absl.testing import absltest
|
2018-11-17 18:03:33 -08:00
|
|
|
from absl.testing import parameterized
|
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
import numpy as np
|
2018-11-17 18:03:33 -08:00
|
|
|
import numpy.random as npr
|
|
|
|
|
2021-04-13 09:42:54 -07:00
|
|
|
from ._src import api
|
2020-02-05 15:38:25 +01:00
|
|
|
from . import core
|
2021-04-07 19:35:17 -07:00
|
|
|
from ._src import dtypes as _dtypes
|
2019-12-10 00:38:18 -08:00
|
|
|
from . import lax
|
2021-04-19 08:52:48 -07:00
|
|
|
from ._src.config import flags, bool_env, config
|
2021-01-11 14:20:32 -08:00
|
|
|
from ._src.util import partial, prod
|
2018-11-17 18:03:33 -08:00
|
|
|
from .tree_util import tree_multimap, tree_all, tree_map, tree_reduce
|
2019-05-10 12:27:15 -07:00
|
|
|
from .lib import xla_bridge
|
2019-09-25 15:59:52 +02:00
|
|
|
from .interpreters import xla
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2018-12-28 13:51:32 -08:00
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
FLAGS = flags.FLAGS
|
|
|
|
flags.DEFINE_enum(
|
2019-05-10 12:27:15 -07:00
|
|
|
'jax_test_dut', '',
|
|
|
|
enum_values=['', 'cpu', 'gpu', 'tpu'],
|
2018-11-17 18:03:33 -08:00
|
|
|
help=
|
|
|
|
'Describes the device under test in case special consideration is required.'
|
|
|
|
)
|
|
|
|
|
2018-12-06 17:31:52 -05:00
|
|
|
flags.DEFINE_integer(
|
|
|
|
'num_generated_cases',
|
2019-04-12 09:29:46 -04:00
|
|
|
int(os.getenv('JAX_NUM_GENERATED_CASES', 10)),
|
2018-12-06 17:31:52 -05:00
|
|
|
help='Number of generated cases to test')
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2021-04-13 10:27:48 +00:00
|
|
|
flags.DEFINE_integer(
|
|
|
|
'max_cases_sampling_retries',
|
|
|
|
int(os.getenv('JAX_MAX_CASES_SAMPLING_RETRIES', 100)),
|
|
|
|
'Number of times a failed test sample should be retried. '
|
|
|
|
'When an unseen case cannot be generated in this many trials, the '
|
|
|
|
'sampling process is terminated.'
|
|
|
|
)
|
|
|
|
|
2020-02-05 17:35:46 +01:00
|
|
|
flags.DEFINE_bool(
|
|
|
|
'jax_skip_slow_tests',
|
2020-02-06 17:27:46 +01:00
|
|
|
bool_env('JAX_SKIP_SLOW_TESTS', False),
|
2020-07-01 11:26:44 -07:00
|
|
|
help='Skip tests marked as slow (> 5 sec).'
|
2020-02-05 17:35:46 +01:00
|
|
|
)
|
|
|
|
|
2020-06-24 16:24:33 -07:00
|
|
|
flags.DEFINE_string(
|
|
|
|
'test_targets', '',
|
|
|
|
'Regular expression specifying which tests to run, called via re.match on '
|
|
|
|
'the test name. If empty or unspecified, run all tests.'
|
|
|
|
)
|
2020-09-29 07:57:20 -07:00
|
|
|
flags.DEFINE_string(
|
|
|
|
'exclude_test_targets', '',
|
|
|
|
'Regular expression specifying which tests NOT to run, called via re.match '
|
|
|
|
'on the test name. If empty or unspecified, run all tests.'
|
|
|
|
)
|
2020-06-24 16:24:33 -07:00
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
EPS = 1e-4
|
|
|
|
|
Change scalar promotion rules to prefer array types over scalar types. (#1709)
* Change scalar promotion rules to prefer array types over scalar types.
Currently JAX does not treat Python scalars specially during type promotion. This means that, for example:
`1. + np.array([...], np.float32)`
ends up as an array of type np.float64. The `1.` is promoted to a default type (here np.float64), and the type promotion of a np.float64 and an np.float32 is an np.float64. This is unlike classic NumPy, which treats scalars specially during type promotion, in particular, preferring the type of an array over the type of a scalar.
This change adds a notion of weak_type to JAX avals. During type promotion, we prefer non-weak types, i.e., the type of the array in the example above, ignoring the type of the scalar.
In contexts where a Python scalar is to be promoted to a NumPy value, a default type is used (e.g., `np.float_`). This change also makes it possible to use 32-bit default types that differ from NumPy's default types. The JAX test suite passes with 32-bit default types. However, we do not yet enable this change or expose it in the API.
2019-11-18 14:51:10 -05:00
|
|
|
def _dtype(x):
|
|
|
|
return (getattr(x, 'dtype', None) or
|
2020-07-07 17:01:38 -07:00
|
|
|
np.dtype(_dtypes.python_scalar_dtypes.get(type(x), None)) or
|
2020-05-12 21:37:05 -03:00
|
|
|
np.asarray(x).dtype)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-06-29 16:22:05 -07:00
|
|
|
|
|
|
|
def num_float_bits(dtype):
|
2020-07-07 17:01:38 -07:00
|
|
|
return _dtypes.finfo(_dtypes.canonicalize_dtype(dtype)).bits
|
2020-06-29 16:22:05 -07:00
|
|
|
|
|
|
|
|
2019-04-24 21:31:15 -07:00
|
|
|
def is_sequence(x):
|
|
|
|
try:
|
|
|
|
iter(x)
|
|
|
|
except TypeError:
|
|
|
|
return False
|
|
|
|
else:
|
|
|
|
return True
|
|
|
|
|
2019-11-20 22:43:46 -05:00
|
|
|
_default_tolerance = {
|
2020-05-12 21:37:05 -03:00
|
|
|
np.dtype(np.bool_): 0,
|
|
|
|
np.dtype(np.int8): 0,
|
|
|
|
np.dtype(np.int16): 0,
|
|
|
|
np.dtype(np.int32): 0,
|
|
|
|
np.dtype(np.int64): 0,
|
|
|
|
np.dtype(np.uint8): 0,
|
|
|
|
np.dtype(np.uint16): 0,
|
|
|
|
np.dtype(np.uint32): 0,
|
|
|
|
np.dtype(np.uint64): 0,
|
2020-07-07 17:01:38 -07:00
|
|
|
np.dtype(_dtypes.bfloat16): 1e-2,
|
2020-05-12 21:37:05 -03:00
|
|
|
np.dtype(np.float16): 1e-3,
|
|
|
|
np.dtype(np.float32): 1e-6,
|
|
|
|
np.dtype(np.float64): 1e-15,
|
|
|
|
np.dtype(np.complex64): 1e-6,
|
|
|
|
np.dtype(np.complex128): 1e-15,
|
2019-11-16 13:51:42 -05:00
|
|
|
}
|
|
|
|
|
2019-11-20 22:43:46 -05:00
|
|
|
def default_tolerance():
|
|
|
|
if device_under_test() != "tpu":
|
|
|
|
return _default_tolerance
|
|
|
|
tol = _default_tolerance.copy()
|
2020-05-12 21:37:05 -03:00
|
|
|
tol[np.dtype(np.float32)] = 1e-3
|
|
|
|
tol[np.dtype(np.complex64)] = 1e-3
|
2019-11-20 22:43:46 -05:00
|
|
|
return tol
|
2019-11-16 13:51:42 -05:00
|
|
|
|
|
|
|
default_gradient_tolerance = {
|
2020-07-07 17:01:38 -07:00
|
|
|
np.dtype(_dtypes.bfloat16): 1e-1,
|
2020-05-12 21:37:05 -03:00
|
|
|
np.dtype(np.float16): 1e-2,
|
|
|
|
np.dtype(np.float32): 2e-3,
|
|
|
|
np.dtype(np.float64): 1e-5,
|
|
|
|
np.dtype(np.complex64): 1e-3,
|
|
|
|
np.dtype(np.complex128): 1e-5,
|
2019-11-16 13:51:42 -05:00
|
|
|
}
|
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
def _assert_numpy_allclose(a, b, atol=None, rtol=None, err_msg=''):
|
2020-07-07 17:01:38 -07:00
|
|
|
a = a.astype(np.float32) if a.dtype == _dtypes.bfloat16 else a
|
|
|
|
b = b.astype(np.float32) if b.dtype == _dtypes.bfloat16 else b
|
2019-11-20 22:43:46 -05:00
|
|
|
kw = {}
|
|
|
|
if atol: kw["atol"] = atol
|
|
|
|
if rtol: kw["rtol"] = rtol
|
2021-05-10 17:44:18 -04:00
|
|
|
with np.errstate(invalid='ignore'):
|
|
|
|
# TODO(phawkins): surprisingly, assert_allclose sometimes reports invalid
|
|
|
|
# value errors. It should not do that.
|
|
|
|
np.testing.assert_allclose(a, b, **kw, err_msg=err_msg)
|
2019-11-16 13:51:42 -05:00
|
|
|
|
|
|
|
def tolerance(dtype, tol=None):
|
2020-03-11 16:19:46 -04:00
|
|
|
tol = {} if tol is None else tol
|
2019-11-16 13:51:42 -05:00
|
|
|
if not isinstance(tol, dict):
|
|
|
|
return tol
|
2020-05-12 21:37:05 -03:00
|
|
|
tol = {np.dtype(key): value for key, value in tol.items()}
|
2020-07-07 17:01:38 -07:00
|
|
|
dtype = _dtypes.canonicalize_dtype(np.dtype(dtype))
|
2019-11-20 22:43:46 -05:00
|
|
|
return tol.get(dtype, default_tolerance()[dtype])
|
|
|
|
|
|
|
|
def _normalize_tolerance(tol):
|
|
|
|
tol = tol or 0
|
|
|
|
if isinstance(tol, dict):
|
2020-05-12 21:37:05 -03:00
|
|
|
return {np.dtype(k): v for k, v in tol.items()}
|
2019-11-20 22:43:46 -05:00
|
|
|
else:
|
2020-09-17 21:51:18 +05:30
|
|
|
return {k: tol for k in _default_tolerance}
|
2019-11-20 22:43:46 -05:00
|
|
|
|
|
|
|
def join_tolerance(tol1, tol2):
|
|
|
|
tol1 = _normalize_tolerance(tol1)
|
|
|
|
tol2 = _normalize_tolerance(tol2)
|
|
|
|
out = tol1
|
|
|
|
for k, v in tol2.items():
|
|
|
|
out[k] = max(v, tol1.get(k, 0))
|
|
|
|
return out
|
2019-11-16 13:51:42 -05:00
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
def _assert_numpy_close(a, b, atol=None, rtol=None, err_msg=''):
|
2018-12-14 18:40:50 -08:00
|
|
|
assert a.shape == b.shape
|
2019-11-16 13:51:42 -05:00
|
|
|
atol = max(tolerance(a.dtype, atol), tolerance(b.dtype, atol))
|
|
|
|
rtol = max(tolerance(a.dtype, rtol), tolerance(b.dtype, rtol))
|
2021-01-13 17:29:47 -08:00
|
|
|
_assert_numpy_allclose(a, b, atol=atol * a.size, rtol=rtol * b.size,
|
|
|
|
err_msg=err_msg)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
def check_eq(xs, ys, err_msg=''):
|
|
|
|
assert_close = partial(_assert_numpy_allclose, err_msg=err_msg)
|
|
|
|
tree_all(tree_multimap(assert_close, xs, ys))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
def check_close(xs, ys, atol=None, rtol=None, err_msg=''):
|
|
|
|
assert_close = partial(_assert_numpy_close, atol=atol, rtol=rtol,
|
|
|
|
err_msg=err_msg)
|
2019-11-11 12:51:15 -08:00
|
|
|
tree_all(tree_multimap(assert_close, xs, ys))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-06-25 08:14:54 -04:00
|
|
|
def _check_dtypes_match(xs, ys):
|
|
|
|
def _assert_dtypes_match(x, y):
|
2021-02-04 09:48:22 -08:00
|
|
|
if config.x64_enabled:
|
2020-06-25 08:14:54 -04:00
|
|
|
assert _dtype(x) == _dtype(y)
|
|
|
|
else:
|
2020-07-07 17:01:38 -07:00
|
|
|
assert (_dtypes.canonicalize_dtype(_dtype(x)) ==
|
|
|
|
_dtypes.canonicalize_dtype(_dtype(y)))
|
2020-06-25 08:14:54 -04:00
|
|
|
tree_all(tree_multimap(_assert_dtypes_match, xs, ys))
|
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
def inner_prod(xs, ys):
|
2019-11-20 22:43:46 -05:00
|
|
|
def contract(x, y):
|
2020-05-12 21:37:05 -03:00
|
|
|
return np.real(np.dot(np.conj(x).reshape(-1), y.reshape(-1)))
|
|
|
|
return tree_reduce(np.add, tree_multimap(contract, xs, ys))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-07-20 17:27:24 -04:00
|
|
|
def _safe_subtract(x, y, *, dtype):
|
|
|
|
"""Subtraction that with `inf - inf == 0` semantics."""
|
|
|
|
with np.errstate(invalid='ignore'):
|
|
|
|
return np.where(np.equal(x, y), np.array(0, dtype),
|
|
|
|
np.subtract(x, y, dtype=dtype))
|
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
add = partial(tree_multimap, lambda x, y: np.add(x, y, dtype=_dtype(x)))
|
|
|
|
sub = partial(tree_multimap, lambda x, y: np.subtract(x, y, dtype=_dtype(x)))
|
2020-07-20 17:27:24 -04:00
|
|
|
safe_sub = partial(tree_multimap,
|
|
|
|
lambda x, y: _safe_subtract(x, y, dtype=_dtype(x)))
|
2020-05-12 21:37:05 -03:00
|
|
|
conj = partial(tree_map, lambda x: np.conj(x, dtype=_dtype(x)))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
def scalar_mul(xs, a):
|
2020-05-12 21:37:05 -03:00
|
|
|
return tree_map(lambda x: np.multiply(x, a, dtype=_dtype(x)), xs)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
|
|
|
def rand_like(rng, x):
|
2020-05-12 21:37:05 -03:00
|
|
|
shape = np.shape(x)
|
2018-11-17 18:03:33 -08:00
|
|
|
dtype = _dtype(x)
|
2020-05-12 21:37:05 -03:00
|
|
|
randn = lambda: np.asarray(rng.randn(*shape), dtype=dtype)
|
2020-07-07 17:01:38 -07:00
|
|
|
if _dtypes.issubdtype(dtype, np.complexfloating):
|
2019-02-01 14:01:06 -05:00
|
|
|
return randn() + dtype.type(1.0j) * randn()
|
2018-11-17 18:03:33 -08:00
|
|
|
else:
|
|
|
|
return randn()
|
|
|
|
|
|
|
|
|
|
|
|
def numerical_jvp(f, primals, tangents, eps=EPS):
|
2019-02-15 14:01:59 -05:00
|
|
|
delta = scalar_mul(tangents, eps)
|
2018-11-17 18:03:33 -08:00
|
|
|
f_pos = f(*add(primals, delta))
|
|
|
|
f_neg = f(*sub(primals, delta))
|
2020-07-20 17:27:24 -04:00
|
|
|
return scalar_mul(safe_sub(f_pos, f_neg), 0.5 / eps)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
Change scalar promotion rules to prefer array types over scalar types. (#1709)
* Change scalar promotion rules to prefer array types over scalar types.
Currently JAX does not treat Python scalars specially during type promotion. This means that, for example:
`1. + np.array([...], np.float32)`
ends up as an array of type np.float64. The `1.` is promoted to a default type (here np.float64), and the type promotion of a np.float64 and an np.float32 is an np.float64. This is unlike classic NumPy, which treats scalars specially during type promotion, in particular, preferring the type of an array over the type of a scalar.
This change adds a notion of weak_type to JAX avals. During type promotion, we prefer non-weak types, i.e., the type of the array in the example above, ignoring the type of the scalar.
In contexts where a Python scalar is to be promoted to a NumPy value, a default type is used (e.g., `np.float_`). This change also makes it possible to use 32-bit default types that differ from NumPy's default types. The JAX test suite passes with 32-bit default types. However, we do not yet enable this change or expose it in the API.
2019-11-18 14:51:10 -05:00
|
|
|
def _merge_tolerance(tol, default):
|
|
|
|
if tol is None:
|
|
|
|
return default
|
|
|
|
if not isinstance(tol, dict):
|
|
|
|
return tol
|
|
|
|
out = default.copy()
|
|
|
|
for k, v in tol.items():
|
2020-05-12 21:37:05 -03:00
|
|
|
out[np.dtype(k)] = v
|
Change scalar promotion rules to prefer array types over scalar types. (#1709)
* Change scalar promotion rules to prefer array types over scalar types.
Currently JAX does not treat Python scalars specially during type promotion. This means that, for example:
`1. + np.array([...], np.float32)`
ends up as an array of type np.float64. The `1.` is promoted to a default type (here np.float64), and the type promotion of a np.float64 and an np.float32 is an np.float64. This is unlike classic NumPy, which treats scalars specially during type promotion, in particular, preferring the type of an array over the type of a scalar.
This change adds a notion of weak_type to JAX avals. During type promotion, we prefer non-weak types, i.e., the type of the array in the example above, ignoring the type of the scalar.
In contexts where a Python scalar is to be promoted to a NumPy value, a default type is used (e.g., `np.float_`). This change also makes it possible to use 32-bit default types that differ from NumPy's default types. The JAX test suite passes with 32-bit default types. However, we do not yet enable this change or expose it in the API.
2019-11-18 14:51:10 -05:00
|
|
|
return out
|
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
|
|
|
|
def check_jvp(f, f_jvp, args, atol=None, rtol=None, eps=EPS, err_msg=''):
|
Change scalar promotion rules to prefer array types over scalar types. (#1709)
* Change scalar promotion rules to prefer array types over scalar types.
Currently JAX does not treat Python scalars specially during type promotion. This means that, for example:
`1. + np.array([...], np.float32)`
ends up as an array of type np.float64. The `1.` is promoted to a default type (here np.float64), and the type promotion of a np.float64 and an np.float32 is an np.float64. This is unlike classic NumPy, which treats scalars specially during type promotion, in particular, preferring the type of an array over the type of a scalar.
This change adds a notion of weak_type to JAX avals. During type promotion, we prefer non-weak types, i.e., the type of the array in the example above, ignoring the type of the scalar.
In contexts where a Python scalar is to be promoted to a NumPy value, a default type is used (e.g., `np.float_`). This change also makes it possible to use 32-bit default types that differ from NumPy's default types. The JAX test suite passes with 32-bit default types. However, we do not yet enable this change or expose it in the API.
2019-11-18 14:51:10 -05:00
|
|
|
atol = _merge_tolerance(atol, default_gradient_tolerance)
|
|
|
|
rtol = _merge_tolerance(rtol, default_gradient_tolerance)
|
2020-05-12 21:37:05 -03:00
|
|
|
rng = np.random.RandomState(0)
|
2018-11-17 18:03:33 -08:00
|
|
|
tangent = tree_map(partial(rand_like, rng), args)
|
|
|
|
v_out, t_out = f_jvp(args, tangent)
|
2020-06-25 08:14:54 -04:00
|
|
|
_check_dtypes_match(v_out, t_out)
|
2018-11-17 18:03:33 -08:00
|
|
|
v_out_expected = f(*args)
|
2020-06-25 08:14:54 -04:00
|
|
|
_check_dtypes_match(v_out, v_out_expected)
|
2018-11-17 18:03:33 -08:00
|
|
|
t_out_expected = numerical_jvp(f, args, tangent, eps=eps)
|
2019-10-22 19:53:59 -04:00
|
|
|
# In principle we should expect exact equality of v_out and v_out_expected,
|
|
|
|
# but due to nondeterminism especially on GPU (e.g., due to convolution
|
|
|
|
# autotuning) we only require "close".
|
2021-01-13 17:29:47 -08:00
|
|
|
check_close(v_out, v_out_expected, atol=atol, rtol=rtol,
|
|
|
|
err_msg=f'{err_msg} primal' if err_msg else 'primal')
|
|
|
|
check_close(t_out, t_out_expected, atol=atol, rtol=rtol,
|
|
|
|
err_msg=f'{err_msg} tangent' if err_msg else 'tangent')
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
def check_vjp(f, f_vjp, args, atol=None, rtol=None, eps=EPS, err_msg=''):
|
Change scalar promotion rules to prefer array types over scalar types. (#1709)
* Change scalar promotion rules to prefer array types over scalar types.
Currently JAX does not treat Python scalars specially during type promotion. This means that, for example:
`1. + np.array([...], np.float32)`
ends up as an array of type np.float64. The `1.` is promoted to a default type (here np.float64), and the type promotion of a np.float64 and an np.float32 is an np.float64. This is unlike classic NumPy, which treats scalars specially during type promotion, in particular, preferring the type of an array over the type of a scalar.
This change adds a notion of weak_type to JAX avals. During type promotion, we prefer non-weak types, i.e., the type of the array in the example above, ignoring the type of the scalar.
In contexts where a Python scalar is to be promoted to a NumPy value, a default type is used (e.g., `np.float_`). This change also makes it possible to use 32-bit default types that differ from NumPy's default types. The JAX test suite passes with 32-bit default types. However, we do not yet enable this change or expose it in the API.
2019-11-18 14:51:10 -05:00
|
|
|
atol = _merge_tolerance(atol, default_gradient_tolerance)
|
|
|
|
rtol = _merge_tolerance(rtol, default_gradient_tolerance)
|
2020-05-12 21:37:05 -03:00
|
|
|
_rand_like = partial(rand_like, np.random.RandomState(0))
|
2018-11-17 18:03:33 -08:00
|
|
|
v_out, vjpfun = f_vjp(*args)
|
|
|
|
v_out_expected = f(*args)
|
2021-01-13 17:29:47 -08:00
|
|
|
check_close(v_out, v_out_expected, atol=atol, rtol=rtol,
|
|
|
|
err_msg=f'{err_msg} primal' if err_msg else 'primal')
|
2018-11-17 18:03:33 -08:00
|
|
|
tangent = tree_map(_rand_like, args)
|
2019-02-15 14:01:59 -05:00
|
|
|
tangent_out = numerical_jvp(f, args, tangent, eps=eps)
|
2018-11-17 18:03:33 -08:00
|
|
|
cotangent = tree_map(_rand_like, v_out)
|
|
|
|
cotangent_out = conj(vjpfun(conj(cotangent)))
|
|
|
|
ip = inner_prod(tangent, cotangent_out)
|
|
|
|
ip_expected = inner_prod(tangent_out, cotangent)
|
2021-01-13 17:29:47 -08:00
|
|
|
check_close(ip, ip_expected, atol=atol, rtol=rtol,
|
|
|
|
err_msg=(f'{err_msg} cotangent projection'
|
|
|
|
if err_msg else 'cotangent projection'))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2019-05-21 17:22:33 -07:00
|
|
|
def check_grads(f, args, order,
|
|
|
|
modes=["fwd", "rev"], atol=None, rtol=None, eps=None):
|
2020-04-09 10:18:07 -07:00
|
|
|
"""Check gradients from automatic differentiation against finite differences.
|
|
|
|
|
|
|
|
Gradients are only checked in a single randomly chosen direction, which
|
|
|
|
ensures that the finite difference calculation does not become prohibitively
|
|
|
|
expensive even for large input/output spaces.
|
|
|
|
|
|
|
|
Args:
|
|
|
|
f: function to check at ``f(*args)``.
|
|
|
|
args: tuple of argument values.
|
|
|
|
order: forward and backwards gradients up to this order are checked.
|
|
|
|
modes: lists of gradient modes to check ('fwd' and/or 'rev').
|
|
|
|
atol: absolute tolerance for gradient equality.
|
|
|
|
rtol: relative tolerance for gradient equality.
|
|
|
|
eps: step size used for finite differences.
|
|
|
|
|
|
|
|
Raises:
|
|
|
|
AssertionError: if gradients do not match.
|
|
|
|
"""
|
2019-05-03 12:37:14 -07:00
|
|
|
args = tuple(args)
|
2019-05-21 17:22:33 -07:00
|
|
|
eps = eps or EPS
|
2019-05-21 15:14:28 -07:00
|
|
|
|
2019-05-21 17:22:33 -07:00
|
|
|
_check_jvp = partial(check_jvp, atol=atol, rtol=rtol, eps=eps)
|
|
|
|
_check_vjp = partial(check_vjp, atol=atol, rtol=rtol, eps=eps)
|
2019-05-21 15:14:28 -07:00
|
|
|
|
2021-01-13 17:29:47 -08:00
|
|
|
def _check_grads(f, args, order, err_msg=''):
|
2019-05-21 15:14:28 -07:00
|
|
|
if "fwd" in modes:
|
2021-01-13 17:29:47 -08:00
|
|
|
fwd_msg = f'JVP of {err_msg}' if err_msg else 'JVP'
|
|
|
|
_check_jvp(f, partial(api.jvp, f), args, err_msg=fwd_msg)
|
2019-05-21 17:22:33 -07:00
|
|
|
if order > 1:
|
2021-01-13 17:29:47 -08:00
|
|
|
_check_grads(partial(api.jvp, f), (args, args), order - 1, fwd_msg)
|
2019-05-21 17:22:33 -07:00
|
|
|
|
2019-05-21 15:14:28 -07:00
|
|
|
if "rev" in modes:
|
2021-01-13 17:29:47 -08:00
|
|
|
rev_msg = f'VJP of {err_msg}' if err_msg else 'VJP'
|
|
|
|
_check_vjp(f, partial(api.vjp, f), args, err_msg=rev_msg)
|
2019-05-21 17:22:33 -07:00
|
|
|
if order > 1:
|
|
|
|
def f_vjp(*args):
|
|
|
|
out_primal_py, vjp_py = api.vjp(f, *args)
|
|
|
|
return vjp_py(out_primal_py)
|
2021-01-13 17:29:47 -08:00
|
|
|
_check_grads(f_vjp, args, order - 1, rev_msg)
|
2019-05-21 17:22:33 -07:00
|
|
|
|
|
|
|
_check_grads(f, args, order)
|
2018-12-17 17:20:52 -08:00
|
|
|
|
2019-12-19 11:19:58 -08:00
|
|
|
|
2021-03-09 14:22:27 -08:00
|
|
|
@contextmanager
|
|
|
|
def count_device_put():
|
|
|
|
device_put = xla.device_put
|
|
|
|
count = [0]
|
|
|
|
|
|
|
|
def device_put_and_count(*args, **kwargs):
|
|
|
|
count[0] += 1
|
|
|
|
return device_put(*args, **kwargs)
|
|
|
|
|
|
|
|
xla.device_put = device_put_and_count
|
|
|
|
try:
|
|
|
|
yield count
|
|
|
|
finally:
|
|
|
|
xla.device_put = device_put
|
|
|
|
|
|
|
|
|
2019-12-19 11:19:58 -08:00
|
|
|
@contextmanager
|
|
|
|
def count_primitive_compiles():
|
|
|
|
xla.xla_primitive_callable.cache_clear()
|
|
|
|
|
|
|
|
# We count how many times we call primitive_computation (which is called
|
|
|
|
# inside xla_primitive_callable) instead of xla_primitive_callable so we don't
|
|
|
|
# count cache hits.
|
|
|
|
primitive_computation = xla.primitive_computation
|
|
|
|
count = [0]
|
|
|
|
|
|
|
|
def primitive_computation_and_count(*args, **kwargs):
|
|
|
|
count[0] += 1
|
|
|
|
return primitive_computation(*args, **kwargs)
|
|
|
|
|
|
|
|
xla.primitive_computation = primitive_computation_and_count
|
|
|
|
try:
|
|
|
|
yield count
|
|
|
|
finally:
|
|
|
|
xla.primitive_computation = primitive_computation
|
|
|
|
|
|
|
|
|
|
|
|
@contextmanager
|
|
|
|
def count_jit_and_pmap_compiles():
|
|
|
|
# No need to clear any caches since we generally jit and pmap fresh callables
|
|
|
|
# in tests.
|
|
|
|
|
|
|
|
jaxpr_subcomp = xla.jaxpr_subcomp
|
|
|
|
count = [0]
|
|
|
|
|
|
|
|
def jaxpr_subcomp_and_count(*args, **kwargs):
|
|
|
|
count[0] += 1
|
|
|
|
return jaxpr_subcomp(*args, **kwargs)
|
|
|
|
|
|
|
|
xla.jaxpr_subcomp = jaxpr_subcomp_and_count
|
|
|
|
try:
|
|
|
|
yield count
|
|
|
|
finally:
|
|
|
|
xla.jaxpr_subcomp = jaxpr_subcomp
|
|
|
|
|
2020-12-02 14:13:05 +00:00
|
|
|
@contextmanager
|
|
|
|
def assert_num_jit_and_pmap_compilations(times):
|
|
|
|
with count_jit_and_pmap_compiles() as count:
|
|
|
|
yield
|
|
|
|
if count[0] != times:
|
|
|
|
raise AssertionError(f"Expected exactly {times} XLA compilations, "
|
|
|
|
f"but executed {count[0]}")
|
2019-12-19 11:19:58 -08:00
|
|
|
|
2019-08-04 17:17:49 -04:00
|
|
|
def device_under_test():
|
|
|
|
return FLAGS.jax_test_dut or xla_bridge.get_backend().platform
|
2018-12-17 17:20:52 -08:00
|
|
|
|
2020-03-19 08:54:37 +01:00
|
|
|
def if_device_under_test(device_type: Union[str, Sequence[str]],
|
|
|
|
if_true, if_false):
|
|
|
|
"""Chooses `if_true` of `if_false` based on device_under_test."""
|
|
|
|
if device_under_test() in ([device_type] if isinstance(device_type, str)
|
|
|
|
else device_type):
|
|
|
|
return if_true
|
|
|
|
else:
|
|
|
|
return if_false
|
|
|
|
|
2019-10-22 19:53:59 -04:00
|
|
|
def supported_dtypes():
|
|
|
|
if device_under_test() == "tpu":
|
2020-08-12 10:02:35 -04:00
|
|
|
types = {np.bool_, np.int8, np.int16, np.int32, np.uint8, np.uint16,
|
|
|
|
np.uint32, _dtypes.bfloat16, np.float16, np.float32, np.complex64}
|
2019-10-22 19:53:59 -04:00
|
|
|
else:
|
2020-07-07 17:01:38 -07:00
|
|
|
types = {np.bool_, np.int8, np.int16, np.int32, np.int64,
|
|
|
|
np.uint8, np.uint16, np.uint32, np.uint64,
|
|
|
|
_dtypes.bfloat16, np.float16, np.float32, np.float64,
|
|
|
|
np.complex64, np.complex128}
|
2021-02-04 09:48:22 -08:00
|
|
|
if not config.x64_enabled:
|
2020-07-07 17:01:38 -07:00
|
|
|
types -= {np.uint64, np.int64, np.float64, np.complex128}
|
|
|
|
return types
|
2019-10-22 19:53:59 -04:00
|
|
|
|
2020-04-02 11:13:40 +02:00
|
|
|
def skip_if_unsupported_type(dtype):
|
2020-07-30 11:07:56 -07:00
|
|
|
dtype = np.dtype(dtype)
|
|
|
|
if dtype.type not in supported_dtypes():
|
2020-04-12 15:35:35 -04:00
|
|
|
raise unittest.SkipTest(
|
2020-07-30 11:07:56 -07:00
|
|
|
f"Type {dtype.name} not supported on {device_under_test()}")
|
2020-04-02 11:13:40 +02:00
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
def skip_on_devices(*disabled_devices):
|
|
|
|
"""A decorator for test methods to skip the test on certain devices."""
|
|
|
|
def skip(test_method):
|
|
|
|
@functools.wraps(test_method)
|
|
|
|
def test_method_wrapper(self, *args, **kwargs):
|
2019-08-04 17:17:49 -04:00
|
|
|
device = device_under_test()
|
2018-11-17 18:03:33 -08:00
|
|
|
if device in disabled_devices:
|
|
|
|
test_name = getattr(test_method, '__name__', '[unknown test]')
|
2020-04-12 15:35:35 -04:00
|
|
|
raise unittest.SkipTest(
|
|
|
|
f"{test_name} not supported on {device.upper()}.")
|
2018-11-17 18:03:33 -08:00
|
|
|
return test_method(self, *args, **kwargs)
|
|
|
|
return test_method_wrapper
|
|
|
|
return skip
|
|
|
|
|
2020-12-13 10:44:20 +02:00
|
|
|
def set_host_platform_device_count(nr_devices: int):
|
|
|
|
"""Returns a closure that undoes the operation."""
|
|
|
|
prev_xla_flags = os.getenv("XLA_FLAGS")
|
|
|
|
flags_str = prev_xla_flags or ""
|
|
|
|
# Don't override user-specified device count, or other XLA flags.
|
|
|
|
if "xla_force_host_platform_device_count" not in flags_str:
|
|
|
|
os.environ["XLA_FLAGS"] = (flags_str +
|
|
|
|
f" --xla_force_host_platform_device_count={nr_devices}")
|
|
|
|
# Clear any cached backends so new CPU backend will pick up the env var.
|
|
|
|
xla_bridge.get_backend.cache_clear()
|
|
|
|
def undo():
|
|
|
|
if prev_xla_flags is None:
|
|
|
|
del os.environ["XLA_FLAGS"]
|
|
|
|
else:
|
|
|
|
os.environ["XLA_FLAGS"] = prev_xla_flags
|
|
|
|
xla_bridge.get_backend.cache_clear()
|
|
|
|
return undo
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2018-11-25 18:53:48 -08:00
|
|
|
def skip_on_flag(flag_name, skip_value):
|
|
|
|
"""A decorator for test methods to skip the test when flags are set."""
|
|
|
|
def skip(test_method): # pylint: disable=missing-docstring
|
|
|
|
@functools.wraps(test_method)
|
|
|
|
def test_method_wrapper(self, *args, **kwargs):
|
2021-03-19 13:49:38 -07:00
|
|
|
flag_value = config._read(flag_name)
|
2018-11-25 18:53:48 -08:00
|
|
|
if flag_value == skip_value:
|
|
|
|
test_name = getattr(test_method, '__name__', '[unknown test]')
|
2020-04-12 15:35:35 -04:00
|
|
|
raise unittest.SkipTest(
|
|
|
|
f"{test_name} not supported when FLAGS.{flag_name} is {flag_value}")
|
2018-11-25 18:53:48 -08:00
|
|
|
return test_method(self, *args, **kwargs)
|
|
|
|
return test_method_wrapper
|
|
|
|
return skip
|
|
|
|
|
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
def format_test_name_suffix(opname, shapes, dtypes):
|
|
|
|
arg_descriptions = (format_shape_dtype_string(shape, dtype)
|
|
|
|
for shape, dtype in zip(shapes, dtypes))
|
|
|
|
return '{}_{}'.format(opname.capitalize(), '_'.join(arg_descriptions))
|
|
|
|
|
|
|
|
|
2019-05-19 12:44:51 -07:00
|
|
|
# We use special symbols, represented as singleton objects, to distinguish
|
|
|
|
# between NumPy scalars, Python scalars, and 0-D arrays.
|
|
|
|
class ScalarShape(object):
|
|
|
|
def __len__(self): return 0
|
|
|
|
class _NumpyScalar(ScalarShape): pass
|
|
|
|
class _PythonScalar(ScalarShape): pass
|
2018-12-06 06:21:38 -08:00
|
|
|
NUMPY_SCALAR_SHAPE = _NumpyScalar()
|
2019-05-19 12:44:51 -07:00
|
|
|
PYTHON_SCALAR_SHAPE = _PythonScalar()
|
2018-12-06 06:21:38 -08:00
|
|
|
|
|
|
|
|
|
|
|
def _dims_of_shape(shape):
|
|
|
|
"""Converts `shape` to a tuple of dimensions."""
|
2019-05-19 12:44:51 -07:00
|
|
|
if type(shape) in (list, tuple):
|
|
|
|
return shape
|
|
|
|
elif isinstance(shape, ScalarShape):
|
|
|
|
return ()
|
2021-02-05 10:07:41 -08:00
|
|
|
elif np.ndim(shape) == 0:
|
|
|
|
return (shape,)
|
2019-05-19 12:44:51 -07:00
|
|
|
else:
|
|
|
|
raise TypeError(type(shape))
|
2018-12-06 06:21:38 -08:00
|
|
|
|
|
|
|
|
|
|
|
def _cast_to_shape(value, shape, dtype):
|
|
|
|
"""Casts `value` to the correct Python type for `shape` and `dtype`."""
|
2019-05-19 12:44:51 -07:00
|
|
|
if shape is NUMPY_SCALAR_SHAPE:
|
|
|
|
# explicitly cast to NumPy scalar in case `value` is a Python scalar.
|
2020-05-12 21:37:05 -03:00
|
|
|
return np.dtype(dtype).type(value)
|
2019-05-19 12:44:51 -07:00
|
|
|
elif shape is PYTHON_SCALAR_SHAPE:
|
|
|
|
# explicitly cast to Python scalar via https://stackoverflow.com/a/11389998
|
2020-05-12 21:37:05 -03:00
|
|
|
return np.asarray(value).item()
|
2019-05-19 12:44:51 -07:00
|
|
|
elif type(shape) in (list, tuple):
|
2020-05-12 21:37:05 -03:00
|
|
|
assert np.shape(value) == tuple(shape)
|
2018-12-06 06:21:38 -08:00
|
|
|
return value
|
2021-02-05 10:07:41 -08:00
|
|
|
elif np.ndim(shape) == 0:
|
|
|
|
assert np.shape(value) == (shape,)
|
|
|
|
return value
|
2018-12-06 06:21:38 -08:00
|
|
|
else:
|
2019-05-19 12:44:51 -07:00
|
|
|
raise TypeError(type(shape))
|
2018-12-06 06:21:38 -08:00
|
|
|
|
|
|
|
|
2019-03-18 14:15:34 -07:00
|
|
|
def dtype_str(dtype):
|
2020-05-12 21:37:05 -03:00
|
|
|
return np.dtype(dtype).name
|
2019-03-18 14:15:34 -07:00
|
|
|
|
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
def format_shape_dtype_string(shape, dtype):
|
2020-05-12 21:37:05 -03:00
|
|
|
if isinstance(shape, np.ndarray):
|
2020-05-04 21:08:34 -04:00
|
|
|
return f'{dtype_str(dtype)}[{shape}]'
|
|
|
|
elif isinstance(shape, list):
|
|
|
|
shape = tuple(shape)
|
|
|
|
return _format_shape_dtype_string(shape, dtype)
|
|
|
|
|
|
|
|
@functools.lru_cache(maxsize=64)
|
|
|
|
def _format_shape_dtype_string(shape, dtype):
|
2019-05-19 12:44:51 -07:00
|
|
|
if shape is NUMPY_SCALAR_SHAPE:
|
2019-03-18 14:15:34 -07:00
|
|
|
return dtype_str(dtype)
|
2019-05-19 12:44:51 -07:00
|
|
|
elif shape is PYTHON_SCALAR_SHAPE:
|
|
|
|
return 'py' + dtype_str(dtype)
|
2020-05-04 21:08:34 -04:00
|
|
|
elif type(shape) is tuple:
|
2018-11-17 18:03:33 -08:00
|
|
|
shapestr = ','.join(str(dim) for dim in shape)
|
2019-05-19 12:44:51 -07:00
|
|
|
return '{}[{}]'.format(dtype_str(dtype), shapestr)
|
|
|
|
elif type(shape) is int:
|
|
|
|
return '{}[{},]'.format(dtype_str(dtype), shape)
|
|
|
|
else:
|
|
|
|
raise TypeError(type(shape))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
|
|
|
def _rand_dtype(rand, shape, dtype, scale=1., post=lambda x: x):
|
|
|
|
"""Produce random values given shape, dtype, scale, and post-processor.
|
|
|
|
|
|
|
|
Args:
|
|
|
|
rand: a function for producing random values of a given shape, e.g. a
|
2020-05-12 21:37:05 -03:00
|
|
|
bound version of either np.RandomState.randn or np.RandomState.rand.
|
2018-11-17 18:03:33 -08:00
|
|
|
shape: a shape value as a tuple of positive integers.
|
|
|
|
dtype: a numpy dtype.
|
|
|
|
scale: optional, a multiplicative scale for the random values (default 1).
|
|
|
|
post: optional, a callable for post-processing the random values (default
|
|
|
|
identity).
|
|
|
|
|
|
|
|
Returns:
|
|
|
|
An ndarray of the given shape and dtype using random values based on a call
|
|
|
|
to rand but scaled, converted to the appropriate dtype, and post-processed.
|
|
|
|
"""
|
2020-05-12 21:37:05 -03:00
|
|
|
r = lambda: np.asarray(scale * rand(*_dims_of_shape(shape)), dtype)
|
2020-07-07 17:01:38 -07:00
|
|
|
if _dtypes.issubdtype(dtype, np.complexfloating):
|
2018-11-17 18:03:33 -08:00
|
|
|
vals = r() + 1.0j * r()
|
|
|
|
else:
|
|
|
|
vals = r()
|
2020-05-12 21:37:05 -03:00
|
|
|
return _cast_to_shape(np.asarray(post(vals), dtype), shape, dtype)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-21 09:20:59 -07:00
|
|
|
def rand_fullrange(rng, standardize_nans=False):
|
2020-05-21 06:40:24 -07:00
|
|
|
"""Random numbers that span the full range of available bits."""
|
|
|
|
def gen(shape, dtype, post=lambda x: x):
|
|
|
|
dtype = np.dtype(dtype)
|
|
|
|
size = dtype.itemsize * np.prod(_dims_of_shape(shape))
|
|
|
|
vals = rng.randint(0, np.iinfo(np.uint8).max, size=size, dtype=np.uint8)
|
|
|
|
vals = post(vals).view(dtype).reshape(shape)
|
|
|
|
# Non-standard NaNs cause errors in numpy equality assertions.
|
|
|
|
if standardize_nans and np.issubdtype(dtype, np.floating):
|
|
|
|
vals[np.isnan(vals)] = np.nan
|
|
|
|
return _cast_to_shape(vals, shape, dtype)
|
|
|
|
return gen
|
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_default(rng, scale=3):
|
|
|
|
return partial(_rand_dtype, rng.randn, scale=scale)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_nonzero(rng):
|
2020-05-12 21:37:05 -03:00
|
|
|
post = lambda x: np.where(x == 0, np.array(1, dtype=x.dtype), x)
|
2020-05-04 23:00:20 -04:00
|
|
|
return partial(_rand_dtype, rng.randn, scale=3, post=post)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_positive(rng):
|
2018-11-17 18:03:33 -08:00
|
|
|
post = lambda x: x + 1
|
2020-05-04 23:00:20 -04:00
|
|
|
return partial(_rand_dtype, rng.rand, scale=2, post=post)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_small(rng):
|
|
|
|
return partial(_rand_dtype, rng.randn, scale=1e-3)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_not_small(rng, offset=10.):
|
2020-05-12 21:37:05 -03:00
|
|
|
post = lambda x: x + np.where(x > 0, offset, -offset)
|
2020-05-04 23:00:20 -04:00
|
|
|
return partial(_rand_dtype, rng.randn, scale=3., post=post)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_small_positive(rng):
|
|
|
|
return partial(_rand_dtype, rng.rand, scale=2e-5)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_uniform(rng, low=0.0, high=1.0):
|
2019-01-14 15:28:53 -05:00
|
|
|
assert low < high
|
2019-01-14 16:39:30 -05:00
|
|
|
post = lambda x: x * (high - low) + low
|
2020-05-04 23:00:20 -04:00
|
|
|
return partial(_rand_dtype, rng.rand, post=post)
|
2019-01-14 15:28:53 -05:00
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_some_equal(rng):
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
def post(x):
|
2018-12-10 08:42:11 -05:00
|
|
|
x_ravel = x.ravel()
|
|
|
|
if len(x_ravel) == 0:
|
|
|
|
return x
|
2020-05-12 21:37:05 -03:00
|
|
|
flips = rng.rand(*np.shape(x)) < 0.5
|
|
|
|
return np.where(flips, x_ravel[0], x)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
return partial(_rand_dtype, rng.randn, scale=100., post=post)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_some_inf(rng):
|
2018-11-17 18:03:33 -08:00
|
|
|
"""Return a random sampler that produces infinities in floating types."""
|
2020-05-04 23:00:20 -04:00
|
|
|
base_rand = rand_default(rng)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2019-09-23 10:04:31 -07:00
|
|
|
"""
|
|
|
|
TODO: Complex numbers are not correctly tested
|
|
|
|
If blocks should be switched in order, and relevant tests should be fixed
|
|
|
|
"""
|
2018-11-17 18:03:33 -08:00
|
|
|
def rand(shape, dtype):
|
|
|
|
"""The random sampler function."""
|
2020-07-07 17:01:38 -07:00
|
|
|
if not _dtypes.issubdtype(dtype, np.floating):
|
2018-11-17 18:03:33 -08:00
|
|
|
# only float types have inf
|
|
|
|
return base_rand(shape, dtype)
|
|
|
|
|
2020-07-07 17:01:38 -07:00
|
|
|
if _dtypes.issubdtype(dtype, np.complexfloating):
|
2020-05-12 21:37:05 -03:00
|
|
|
base_dtype = np.real(np.array(0, dtype=dtype)).dtype
|
2019-12-06 14:49:27 -05:00
|
|
|
out = (rand(shape, base_dtype) +
|
2020-05-12 21:37:05 -03:00
|
|
|
np.array(1j, dtype) * rand(shape, base_dtype))
|
2019-12-06 14:49:27 -05:00
|
|
|
return _cast_to_shape(out, shape, dtype)
|
2019-05-07 15:05:37 -04:00
|
|
|
|
2018-12-06 06:21:38 -08:00
|
|
|
dims = _dims_of_shape(shape)
|
|
|
|
posinf_flips = rng.rand(*dims) < 0.1
|
|
|
|
neginf_flips = rng.rand(*dims) < 0.1
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
vals = base_rand(shape, dtype)
|
2020-05-12 21:37:05 -03:00
|
|
|
vals = np.where(posinf_flips, np.array(np.inf, dtype=dtype), vals)
|
|
|
|
vals = np.where(neginf_flips, np.array(-np.inf, dtype=dtype), vals)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
return _cast_to_shape(np.asarray(vals, dtype=dtype), shape, dtype)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
return rand
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_some_nan(rng):
|
2019-09-22 21:38:34 -07:00
|
|
|
"""Return a random sampler that produces nans in floating types."""
|
2020-05-04 23:00:20 -04:00
|
|
|
base_rand = rand_default(rng)
|
2019-09-22 21:38:34 -07:00
|
|
|
|
|
|
|
def rand(shape, dtype):
|
|
|
|
"""The random sampler function."""
|
2020-07-07 17:01:38 -07:00
|
|
|
if _dtypes.issubdtype(dtype, np.complexfloating):
|
2020-05-12 21:37:05 -03:00
|
|
|
base_dtype = np.real(np.array(0, dtype=dtype)).dtype
|
2019-12-06 14:49:27 -05:00
|
|
|
out = (rand(shape, base_dtype) +
|
2020-05-12 21:37:05 -03:00
|
|
|
np.array(1j, dtype) * rand(shape, base_dtype))
|
2019-12-06 14:49:27 -05:00
|
|
|
return _cast_to_shape(out, shape, dtype)
|
2019-09-22 21:38:34 -07:00
|
|
|
|
2020-07-07 17:01:38 -07:00
|
|
|
if not _dtypes.issubdtype(dtype, np.floating):
|
2019-09-23 08:53:49 -07:00
|
|
|
# only float types have inf
|
|
|
|
return base_rand(shape, dtype)
|
|
|
|
|
2019-09-22 21:38:34 -07:00
|
|
|
dims = _dims_of_shape(shape)
|
2021-04-14 13:53:59 -07:00
|
|
|
r = rng.rand(*dims)
|
|
|
|
nan_flips = r < 0.1
|
|
|
|
neg_nan_flips = r < 0.05
|
2019-09-22 21:38:34 -07:00
|
|
|
|
|
|
|
vals = base_rand(shape, dtype)
|
2020-05-12 21:37:05 -03:00
|
|
|
vals = np.where(nan_flips, np.array(np.nan, dtype=dtype), vals)
|
2021-04-14 13:53:59 -07:00
|
|
|
vals = np.where(neg_nan_flips, np.array(-np.nan, dtype=dtype), vals)
|
2019-09-22 21:38:34 -07:00
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
return _cast_to_shape(np.asarray(vals, dtype=dtype), shape, dtype)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
return rand
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_some_inf_and_nan(rng):
|
2019-05-07 15:05:37 -04:00
|
|
|
"""Return a random sampler that produces infinities in floating types."""
|
2020-05-04 23:00:20 -04:00
|
|
|
base_rand = rand_default(rng)
|
2019-05-07 15:05:37 -04:00
|
|
|
|
2019-09-23 10:04:31 -07:00
|
|
|
"""
|
|
|
|
TODO: Complex numbers are not correctly tested
|
|
|
|
If blocks should be switched in order, and relevant tests should be fixed
|
|
|
|
"""
|
2019-05-07 15:05:37 -04:00
|
|
|
def rand(shape, dtype):
|
|
|
|
"""The random sampler function."""
|
2020-07-07 17:01:38 -07:00
|
|
|
if not _dtypes.issubdtype(dtype, np.floating):
|
2019-05-07 15:05:37 -04:00
|
|
|
# only float types have inf
|
|
|
|
return base_rand(shape, dtype)
|
|
|
|
|
2020-07-07 17:01:38 -07:00
|
|
|
if _dtypes.issubdtype(dtype, np.complexfloating):
|
2020-05-12 21:37:05 -03:00
|
|
|
base_dtype = np.real(np.array(0, dtype=dtype)).dtype
|
2019-12-06 14:49:27 -05:00
|
|
|
out = (rand(shape, base_dtype) +
|
2020-05-12 21:37:05 -03:00
|
|
|
np.array(1j, dtype) * rand(shape, base_dtype))
|
2019-12-06 14:49:27 -05:00
|
|
|
return _cast_to_shape(out, shape, dtype)
|
2019-05-07 15:05:37 -04:00
|
|
|
|
|
|
|
dims = _dims_of_shape(shape)
|
|
|
|
posinf_flips = rng.rand(*dims) < 0.1
|
|
|
|
neginf_flips = rng.rand(*dims) < 0.1
|
|
|
|
nan_flips = rng.rand(*dims) < 0.1
|
|
|
|
|
|
|
|
vals = base_rand(shape, dtype)
|
2020-05-12 21:37:05 -03:00
|
|
|
vals = np.where(posinf_flips, np.array(np.inf, dtype=dtype), vals)
|
|
|
|
vals = np.where(neginf_flips, np.array(-np.inf, dtype=dtype), vals)
|
|
|
|
vals = np.where(nan_flips, np.array(np.nan, dtype=dtype), vals)
|
2019-05-07 15:05:37 -04:00
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
return _cast_to_shape(np.asarray(vals, dtype=dtype), shape, dtype)
|
2019-05-07 15:05:37 -04:00
|
|
|
|
|
|
|
return rand
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
# TODO(mattjj): doesn't handle complex types
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_some_zero(rng):
|
2018-11-17 18:03:33 -08:00
|
|
|
"""Return a random sampler that produces some zeros."""
|
2020-05-04 23:00:20 -04:00
|
|
|
base_rand = rand_default(rng)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
def rand(shape, dtype):
|
|
|
|
"""The random sampler function."""
|
2018-12-06 06:21:38 -08:00
|
|
|
dims = _dims_of_shape(shape)
|
|
|
|
zeros = rng.rand(*dims) < 0.5
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
vals = base_rand(shape, dtype)
|
2020-05-12 21:37:05 -03:00
|
|
|
vals = np.where(zeros, np.array(0, dtype=dtype), vals)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
return _cast_to_shape(np.asarray(vals, dtype=dtype), shape, dtype)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
return rand
|
|
|
|
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_int(rng, low=0, high=None):
|
2019-01-09 21:26:22 -05:00
|
|
|
def fn(shape, dtype):
|
2020-07-23 16:17:55 -04:00
|
|
|
nonlocal high
|
|
|
|
if low == 0 and high is None:
|
|
|
|
if np.issubdtype(dtype, np.integer):
|
|
|
|
high = np.iinfo(dtype).max
|
|
|
|
else:
|
|
|
|
raise ValueError("rand_int requires an explicit `high` value for "
|
|
|
|
"non-integer types.")
|
2020-05-04 23:00:20 -04:00
|
|
|
return rng.randint(low, high=high, size=shape, dtype=dtype)
|
2019-01-09 21:26:22 -05:00
|
|
|
return fn
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_unique_int(rng, high=None):
|
2020-04-19 11:49:15 -07:00
|
|
|
def fn(shape, dtype):
|
2020-08-18 10:17:38 -07:00
|
|
|
return rng.choice(np.arange(high or prod(shape), dtype=dtype),
|
2020-04-19 11:49:15 -07:00
|
|
|
size=shape, replace=False)
|
|
|
|
return fn
|
|
|
|
|
2020-05-04 23:00:20 -04:00
|
|
|
def rand_bool(rng):
|
2018-12-06 06:21:38 -08:00
|
|
|
def generator(shape, dtype):
|
|
|
|
return _cast_to_shape(rng.rand(*_dims_of_shape(shape)) < 0.5, shape, dtype)
|
|
|
|
return generator
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
def check_raises(thunk, err_type, msg):
|
|
|
|
try:
|
|
|
|
thunk()
|
|
|
|
assert False
|
|
|
|
except err_type as e:
|
2018-12-06 21:47:47 -05:00
|
|
|
assert str(e).startswith(msg), "\n{}\n\n{}\n".format(e, msg)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2019-11-16 13:51:42 -05:00
|
|
|
def check_raises_regexp(thunk, err_type, pattern):
|
|
|
|
try:
|
|
|
|
thunk()
|
|
|
|
assert False
|
|
|
|
except err_type as e:
|
|
|
|
assert re.match(pattern, str(e)), "{}\n\n{}\n".format(e, pattern)
|
|
|
|
|
2019-12-10 00:38:18 -08:00
|
|
|
|
2020-07-17 08:44:47 -04:00
|
|
|
def iter_eqns(jaxpr):
|
2020-02-03 20:58:56 +01:00
|
|
|
# TODO(necula): why doesn't this search in params?
|
2019-12-10 00:38:18 -08:00
|
|
|
for eqn in jaxpr.eqns:
|
|
|
|
yield eqn
|
2020-02-05 15:38:25 +01:00
|
|
|
for subjaxpr in core.subjaxprs(jaxpr):
|
2020-07-17 08:44:47 -04:00
|
|
|
yield from iter_eqns(subjaxpr)
|
2019-12-10 00:38:18 -08:00
|
|
|
|
|
|
|
def assert_dot_precision(expected_precision, fun, *args):
|
|
|
|
jaxpr = api.make_jaxpr(fun)(*args)
|
2020-07-17 08:44:47 -04:00
|
|
|
precisions = [eqn.params['precision'] for eqn in iter_eqns(jaxpr.jaxpr)
|
2019-12-10 00:38:18 -08:00
|
|
|
if eqn.primitive == lax.dot_general_p]
|
|
|
|
for precision in precisions:
|
|
|
|
msg = "Unexpected precision: {} != {}".format(expected_precision, precision)
|
2021-05-12 02:29:51 -07:00
|
|
|
if isinstance(precision, tuple):
|
|
|
|
assert precision[0] == expected_precision, msg
|
|
|
|
assert precision[1] == expected_precision, msg
|
|
|
|
else:
|
|
|
|
assert precision == expected_precision, msg
|
2019-12-10 00:38:18 -08:00
|
|
|
|
|
|
|
|
2020-03-18 17:06:05 -04:00
|
|
|
_CACHED_INDICES: Dict[int, Sequence[int]] = {}
|
2019-11-11 12:51:15 -08:00
|
|
|
|
2018-12-06 18:02:43 -05:00
|
|
|
def cases_from_list(xs):
|
2018-12-06 18:30:59 -05:00
|
|
|
xs = list(xs)
|
2019-11-11 12:51:15 -08:00
|
|
|
n = len(xs)
|
|
|
|
k = min(n, FLAGS.num_generated_cases)
|
|
|
|
# Random sampling for every parameterized test is expensive. Do it once and
|
|
|
|
# cache the result.
|
|
|
|
indices = _CACHED_INDICES.get(n)
|
|
|
|
if indices is None:
|
|
|
|
rng = npr.RandomState(42)
|
|
|
|
_CACHED_INDICES[n] = indices = rng.permutation(n)
|
|
|
|
return [xs[i] for i in indices[:k]]
|
2018-12-06 17:31:52 -05:00
|
|
|
|
2018-12-06 18:02:43 -05:00
|
|
|
def cases_from_gens(*gens):
|
|
|
|
sizes = [1, 3, 10]
|
2018-12-06 18:30:59 -05:00
|
|
|
cases_per_size = int(FLAGS.num_generated_cases / len(sizes)) + 1
|
2018-12-06 18:02:43 -05:00
|
|
|
for size in sizes:
|
2020-01-08 13:17:55 -05:00
|
|
|
for i in range(cases_per_size):
|
2018-12-06 18:02:43 -05:00
|
|
|
yield ('_{}_{}'.format(size, i),) + tuple(gen(size) for gen in gens)
|
2018-12-06 17:31:52 -05:00
|
|
|
|
2021-04-13 10:27:48 +00:00
|
|
|
def named_cases_from_sampler(gen):
|
|
|
|
seen = set()
|
|
|
|
retries = 0
|
|
|
|
rng = npr.RandomState(42)
|
|
|
|
def choose_one(x):
|
|
|
|
if not isinstance(x, (list, tuple)):
|
|
|
|
x = list(x)
|
|
|
|
return [x[rng.randint(len(x))]]
|
|
|
|
while (len(seen) < FLAGS.num_generated_cases and
|
|
|
|
retries < FLAGS.max_cases_sampling_retries):
|
|
|
|
retries += 1
|
|
|
|
cases = list(gen(choose_one))
|
|
|
|
if not cases:
|
|
|
|
continue
|
|
|
|
if len(cases) > 1:
|
|
|
|
raise RuntimeError("Generator is expected to only return a single case when sampling")
|
|
|
|
case = cases[0]
|
|
|
|
if case["testcase_name"] in seen:
|
|
|
|
continue
|
|
|
|
retries = 0
|
|
|
|
seen.add(case["testcase_name"])
|
|
|
|
yield case
|
|
|
|
|
2018-12-06 17:31:52 -05:00
|
|
|
|
2020-06-24 16:24:33 -07:00
|
|
|
class JaxTestLoader(absltest.TestLoader):
|
|
|
|
def getTestCaseNames(self, testCaseClass):
|
|
|
|
names = super().getTestCaseNames(testCaseClass)
|
|
|
|
if FLAGS.test_targets:
|
|
|
|
pattern = re.compile(FLAGS.test_targets)
|
2020-06-29 11:06:10 -07:00
|
|
|
names = [name for name in names
|
|
|
|
if pattern.search(f"{testCaseClass.__name__}.{name}")]
|
2020-09-29 07:57:20 -07:00
|
|
|
if FLAGS.exclude_test_targets:
|
|
|
|
pattern = re.compile(FLAGS.exclude_test_targets)
|
|
|
|
names = [name for name in names
|
|
|
|
if not pattern.search(f"{testCaseClass.__name__}.{name}")]
|
2020-06-24 16:24:33 -07:00
|
|
|
return names
|
|
|
|
|
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
class JaxTestCase(parameterized.TestCase):
|
|
|
|
"""Base class for JAX tests including numerical checks and boilerplate."""
|
|
|
|
|
2020-04-13 09:44:13 -07:00
|
|
|
# TODO(mattjj): this obscures the error messages from failures, figure out how
|
|
|
|
# to re-enable it
|
|
|
|
# def tearDown(self) -> None:
|
|
|
|
# assert core.reset_trace_state()
|
2020-04-02 22:01:43 -07:00
|
|
|
|
2020-05-01 09:16:31 +03:00
|
|
|
def setUp(self):
|
2020-05-04 23:00:20 -04:00
|
|
|
super(JaxTestCase, self).setUp()
|
2021-03-19 13:49:38 -07:00
|
|
|
config.update('jax_enable_checks', True)
|
2020-05-04 23:00:20 -04:00
|
|
|
# We use the adler32 hash for two reasons.
|
|
|
|
# a) it is deterministic run to run, unlike hash() which is randomized.
|
|
|
|
# b) it returns values in int32 range, which RandomState requires.
|
|
|
|
self._rng = npr.RandomState(zlib.adler32(self._testMethodName.encode()))
|
|
|
|
|
|
|
|
def rng(self):
|
|
|
|
return self._rng
|
2020-05-01 09:16:31 +03:00
|
|
|
|
2020-06-01 17:19:23 -04:00
|
|
|
def assertArraysEqual(self, x, y, *, check_dtypes=True):
|
2020-05-15 19:09:43 -07:00
|
|
|
"""Assert that x and y arrays are exactly equal."""
|
|
|
|
if check_dtypes:
|
|
|
|
self.assertDtypesMatch(x, y)
|
2020-09-16 20:29:19 -07:00
|
|
|
np.testing.assert_array_equal(x, y)
|
2020-05-15 19:09:43 -07:00
|
|
|
|
2020-06-01 17:19:23 -04:00
|
|
|
def assertArraysAllClose(self, x, y, *, check_dtypes=True, atol=None,
|
|
|
|
rtol=None):
|
2018-11-17 18:03:33 -08:00
|
|
|
"""Assert that x and y are close (up to numerical tolerances)."""
|
2019-03-14 21:59:31 -04:00
|
|
|
self.assertEqual(x.shape, y.shape)
|
2019-11-16 13:51:42 -05:00
|
|
|
atol = max(tolerance(_dtype(x), atol), tolerance(_dtype(y), atol))
|
|
|
|
rtol = max(tolerance(_dtype(x), rtol), tolerance(_dtype(y), rtol))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2019-11-20 22:43:46 -05:00
|
|
|
_assert_numpy_allclose(x, y, atol=atol, rtol=rtol)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
if check_dtypes:
|
|
|
|
self.assertDtypesMatch(x, y)
|
|
|
|
|
2020-06-01 19:29:26 -04:00
|
|
|
def assertDtypesMatch(self, x, y, *, canonicalize_dtypes=True):
|
2021-02-04 09:48:22 -08:00
|
|
|
if not config.x64_enabled and canonicalize_dtypes:
|
2020-07-07 17:01:38 -07:00
|
|
|
self.assertEqual(_dtypes.canonicalize_dtype(_dtype(x)),
|
|
|
|
_dtypes.canonicalize_dtype(_dtype(y)))
|
2020-06-01 19:29:26 -04:00
|
|
|
else:
|
2019-12-09 21:18:39 -05:00
|
|
|
self.assertEqual(_dtype(x), _dtype(y))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2020-06-01 19:29:26 -04:00
|
|
|
def assertAllClose(self, x, y, *, check_dtypes=True, atol=None, rtol=None,
|
|
|
|
canonicalize_dtypes=True):
|
2018-11-17 18:03:33 -08:00
|
|
|
"""Assert that x and y, either arrays or nested tuples/lists, are close."""
|
2019-04-24 21:31:15 -07:00
|
|
|
if isinstance(x, dict):
|
2019-01-07 08:54:14 -08:00
|
|
|
self.assertIsInstance(y, dict)
|
|
|
|
self.assertEqual(set(x.keys()), set(y.keys()))
|
|
|
|
for k in x.keys():
|
2020-06-01 17:19:23 -04:00
|
|
|
self.assertAllClose(x[k], y[k], check_dtypes=check_dtypes, atol=atol,
|
2020-06-01 19:29:26 -04:00
|
|
|
rtol=rtol, canonicalize_dtypes=canonicalize_dtypes)
|
2019-04-24 21:31:15 -07:00
|
|
|
elif is_sequence(x) and not hasattr(x, '__array__'):
|
|
|
|
self.assertTrue(is_sequence(y) and not hasattr(y, '__array__'))
|
|
|
|
self.assertEqual(len(x), len(y))
|
|
|
|
for x_elt, y_elt in zip(x, y):
|
2020-06-01 17:19:23 -04:00
|
|
|
self.assertAllClose(x_elt, y_elt, check_dtypes=check_dtypes, atol=atol,
|
2020-06-01 19:29:26 -04:00
|
|
|
rtol=rtol, canonicalize_dtypes=canonicalize_dtypes)
|
2020-05-12 21:37:05 -03:00
|
|
|
elif hasattr(x, '__array__') or np.isscalar(x):
|
|
|
|
self.assertTrue(hasattr(y, '__array__') or np.isscalar(y))
|
2019-12-09 21:18:39 -05:00
|
|
|
if check_dtypes:
|
2020-06-01 19:29:26 -04:00
|
|
|
self.assertDtypesMatch(x, y, canonicalize_dtypes=canonicalize_dtypes)
|
2020-05-12 21:37:05 -03:00
|
|
|
x = np.asarray(x)
|
|
|
|
y = np.asarray(y)
|
2019-12-09 21:18:39 -05:00
|
|
|
self.assertArraysAllClose(x, y, check_dtypes=False, atol=atol, rtol=rtol)
|
2019-09-22 15:32:12 -07:00
|
|
|
elif x == y:
|
2019-11-16 13:51:42 -05:00
|
|
|
return
|
2019-04-24 21:31:15 -07:00
|
|
|
else:
|
|
|
|
raise TypeError((type(x), type(y)))
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2019-11-26 13:56:58 +01:00
|
|
|
def assertMultiLineStrippedEqual(self, expected, what):
|
2020-09-14 02:47:28 -07:00
|
|
|
"""Asserts two strings are equal, after dedenting and stripping each line."""
|
|
|
|
expected = textwrap.dedent(expected)
|
|
|
|
what = textwrap.dedent(what)
|
2019-11-26 13:56:58 +01:00
|
|
|
ignore_space_re = re.compile(r'\s*\n\s*')
|
|
|
|
expected_clean = re.sub(ignore_space_re, '\n', expected.strip())
|
|
|
|
what_clean = re.sub(ignore_space_re, '\n', what.strip())
|
2019-11-26 14:05:08 +01:00
|
|
|
self.assertMultiLineEqual(expected_clean, what_clean,
|
2020-02-10 11:40:05 +01:00
|
|
|
msg="Found\n{}\nExpecting\n{}".format(what, expected))
|
2019-11-26 08:18:53 +01:00
|
|
|
|
2020-06-01 17:19:23 -04:00
|
|
|
def _CompileAndCheck(self, fun, args_maker, *, check_dtypes=True,
|
2021-03-09 13:25:38 -08:00
|
|
|
rtol=None, atol=None, check_cache_misses=True):
|
2018-11-17 18:03:33 -08:00
|
|
|
"""Helper method for running JAX compilation and allclose assertions."""
|
|
|
|
args = args_maker()
|
|
|
|
|
|
|
|
def wrapped_fun(*args):
|
|
|
|
self.assertTrue(python_should_be_executing)
|
|
|
|
return fun(*args)
|
|
|
|
|
|
|
|
python_should_be_executing = True
|
|
|
|
python_ans = fun(*args)
|
|
|
|
|
2020-05-12 21:37:05 -03:00
|
|
|
python_shapes = tree_map(lambda x: np.shape(x), python_ans)
|
|
|
|
np_shapes = tree_map(lambda x: np.shape(np.asarray(x)), python_ans)
|
|
|
|
self.assertEqual(python_shapes, np_shapes)
|
2019-12-02 14:43:43 -05:00
|
|
|
|
2019-09-25 15:59:52 +02:00
|
|
|
cache_misses = xla.xla_primitive_callable.cache_info().misses
|
|
|
|
python_ans = fun(*args)
|
2021-03-09 13:25:38 -08:00
|
|
|
if check_cache_misses:
|
|
|
|
self.assertEqual(
|
|
|
|
cache_misses, xla.xla_primitive_callable.cache_info().misses,
|
|
|
|
"Compilation detected during second call of {} in op-by-op "
|
|
|
|
"mode.".format(fun))
|
2019-09-25 15:59:52 +02:00
|
|
|
|
2018-11-17 18:03:33 -08:00
|
|
|
cfun = api.jit(wrapped_fun)
|
|
|
|
python_should_be_executing = True
|
|
|
|
monitored_ans = cfun(*args)
|
|
|
|
|
|
|
|
python_should_be_executing = False
|
|
|
|
compiled_ans = cfun(*args)
|
|
|
|
|
2020-06-01 17:19:23 -04:00
|
|
|
self.assertAllClose(python_ans, monitored_ans, check_dtypes=check_dtypes,
|
|
|
|
atol=atol, rtol=rtol)
|
|
|
|
self.assertAllClose(python_ans, compiled_ans, check_dtypes=check_dtypes,
|
|
|
|
atol=atol, rtol=rtol)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
|
|
|
args = args_maker()
|
|
|
|
|
|
|
|
python_should_be_executing = True
|
|
|
|
python_ans = fun(*args)
|
|
|
|
|
|
|
|
python_should_be_executing = False
|
|
|
|
compiled_ans = cfun(*args)
|
|
|
|
|
2020-06-01 17:19:23 -04:00
|
|
|
self.assertAllClose(python_ans, compiled_ans, check_dtypes=check_dtypes,
|
|
|
|
atol=atol, rtol=rtol)
|
2018-11-17 18:03:33 -08:00
|
|
|
|
2019-03-22 17:09:35 -07:00
|
|
|
def _CheckAgainstNumpy(self, numpy_reference_op, lax_op, args_maker,
|
2020-06-01 19:29:26 -04:00
|
|
|
check_dtypes=True, tol=None,
|
|
|
|
canonicalize_dtypes=True):
|
2018-11-17 18:03:33 -08:00
|
|
|
args = args_maker()
|
2018-12-30 18:07:50 -08:00
|
|
|
lax_ans = lax_op(*args)
|
improve implementation of MVN logpdf (#2481)
fixes #2314
I also added a bit more test coverage, but not a ton: scipy has
different batch shape semantics and default arguments than I might
expect, so I didn't bother to implement those (and left some test cases
commented out).
I ran into this surprising scipy bug:
```python
In [1]: from scipy.stats import multivariate_normal
In [2]: import numpy as np
In [3]: args = [np.array(1., np.float32), np.array(2., np.float64), np.array(3., np.float64)]
In [4]: print([x.shape for x in args])
[(), (), ()]
In [5]: multivariate_normal.logpdf(*args)
Out[5]: -1.6349113442053944
In [6]: print([x.shape for x in args])
[(), (1,), (1, 1)]
```
Mutated arguments! But it depends on dtype promotion:
```python
In [7]: args = [np.array(1., np.float32), np.array(2., np.float32), np.array(3., np.float32)]
In [8]: print([x.shape for x in args])
[(), (), ()]
In [9]: multivariate_normal.logpdf(*args)
Out[9]: -1.6349113442053944
In [10]: print([x.shape for x in args])
[(), (), ()]
```
2020-03-21 15:42:59 -07:00
|
|
|
numpy_ans = numpy_reference_op(*args)
|
2019-11-20 22:43:46 -05:00
|
|
|
self.assertAllClose(numpy_ans, lax_ans, check_dtypes=check_dtypes,
|
2020-06-01 19:29:26 -04:00
|
|
|
atol=tol, rtol=tol,
|
|
|
|
canonicalize_dtypes=canonicalize_dtypes)
|
2020-04-12 15:35:35 -04:00
|
|
|
|
|
|
|
|
2021-02-04 14:53:38 +00:00
|
|
|
class BufferDonationTestCase(JaxTestCase):
|
|
|
|
assertDeleted = lambda self, x: self._assertDeleted(x, True)
|
|
|
|
assertNotDeleted = lambda self, x: self._assertDeleted(x, False)
|
|
|
|
|
|
|
|
def _assertDeleted(self, x, deleted):
|
|
|
|
if hasattr(x, "device_buffer"):
|
|
|
|
self.assertEqual(x.device_buffer.is_deleted(), deleted)
|
|
|
|
else:
|
|
|
|
for buffer in x.device_buffers:
|
|
|
|
self.assertEqual(buffer.is_deleted(), deleted)
|
|
|
|
|
|
|
|
|
2020-04-12 15:35:35 -04:00
|
|
|
@contextmanager
|
|
|
|
def ignore_warning(**kw):
|
|
|
|
with warnings.catch_warnings():
|
|
|
|
warnings.filterwarnings("ignore", **kw)
|
|
|
|
yield
|
2020-07-07 17:01:38 -07:00
|
|
|
|
|
|
|
|
|
|
|
class _cached_property:
|
|
|
|
null = object()
|
|
|
|
|
|
|
|
def __init__(self, method):
|
|
|
|
self._method = method
|
|
|
|
self._value = self.null
|
|
|
|
|
|
|
|
def __get__(self, obj, cls):
|
|
|
|
if self._value is self.null:
|
|
|
|
self._value = self._method(obj)
|
|
|
|
return self._value
|
|
|
|
|
|
|
|
|
|
|
|
class _LazyDtypes:
|
|
|
|
"""A class that unifies lists of supported dtypes.
|
|
|
|
|
|
|
|
These could be module-level constants, but device_under_test() is not always
|
|
|
|
known at import time, so we need to define these lists lazily.
|
|
|
|
"""
|
|
|
|
def supported(self, dtypes):
|
|
|
|
supported = supported_dtypes()
|
|
|
|
return type(dtypes)(d for d in dtypes if d in supported)
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def floating(self):
|
|
|
|
return self.supported([np.float32, np.float64])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def all_floating(self):
|
|
|
|
return self.supported([_dtypes.bfloat16, np.float16, np.float32, np.float64])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def integer(self):
|
|
|
|
return self.supported([np.int32, np.int64])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def all_integer(self):
|
|
|
|
return self.supported([np.int8, np.int16, np.int32, np.int64])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def unsigned(self):
|
|
|
|
return self.supported([np.uint32, np.uint64])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def all_unsigned(self):
|
|
|
|
return self.supported([np.uint8, np.uint16, np.uint32, np.uint64])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def complex(self):
|
|
|
|
return self.supported([np.complex64, np.complex128])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def boolean(self):
|
|
|
|
return self.supported([np.bool_])
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def inexact(self):
|
|
|
|
return self.floating + self.complex
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def all_inexact(self):
|
|
|
|
return self.all_floating + self.complex
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def numeric(self):
|
|
|
|
return self.floating + self.integer + self.unsigned + self.complex
|
|
|
|
|
|
|
|
@_cached_property
|
|
|
|
def all(self):
|
|
|
|
return (self.all_floating + self.all_integer + self.all_unsigned +
|
|
|
|
self.complex + self.boolean)
|
|
|
|
|
|
|
|
|
|
|
|
dtypes = _LazyDtypes()
|