forked from jax-ml/jax
-
Notifications
You must be signed in to change notification settings - Fork 0
/
Copy patharray_interoperability_test.py
391 lines (339 loc) · 13.5 KB
/
array_interoperability_test.py
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
# Copyright 2020 The JAX Authors.
#
# 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.
import unittest
from absl.testing import absltest
import jax
import jax.dlpack
import jax.numpy as jnp
from jax.sharding import PartitionSpec as P
from jax._src import config
from jax._src import test_util as jtu
import numpy as np
config.parse_flags_with_absl()
try:
import cupy
except ImportError:
cupy = None
try:
import tensorflow as tf
tf_version = tuple(
int(x) for x in tf.version.VERSION.split("-")[0].split("."))
except ImportError:
tf = None
dlpack_dtypes = sorted(jax.dlpack.SUPPORTED_DTYPES, key=lambda x: x.__name__)
numpy_dtypes = sorted(
[dt for dt in jax.dlpack.SUPPORTED_DTYPES if dt != jnp.bfloat16],
key=lambda x: x.__name__)
cuda_array_interface_dtypes = [dt for dt in dlpack_dtypes if dt != jnp.bfloat16]
nonempty_nonscalar_array_shapes = [(4,), (3, 4), (2, 3, 4)]
empty_array_shapes = []
empty_array_shapes += [(0,), (0, 4), (3, 0),]
nonempty_nonscalar_array_shapes += [(3, 1), (1, 4), (2, 1, 4)]
nonempty_array_shapes = [()] + nonempty_nonscalar_array_shapes
all_shapes = nonempty_array_shapes + empty_array_shapes
class DLPackTest(jtu.JaxTestCase):
def setUp(self):
super().setUp()
if not jtu.test_device_matches(["cpu", "gpu"]):
self.skipTest(f"DLPack not supported on {jtu.device_under_test()}")
@jtu.sample_product(
shape=all_shapes,
dtype=dlpack_dtypes,
copy=[False, True, None],
use_stream=[False, True],
)
@jtu.run_on_devices("gpu")
@jtu.ignore_warning(message="Calling from_dlpack with a DLPack tensor",
category=DeprecationWarning)
def testJaxRoundTrip(self, shape, dtype, copy, use_stream):
rng = jtu.rand_default(self.rng())
np = rng(shape, dtype)
def _check_copy(x: jax.Array, y: jax.Array, expect_copy):
copied = x.unsafe_buffer_pointer() != y.unsafe_buffer_pointer()
assert copied == expect_copy, f"Expected {'a' if expect_copy else 'no'} copy"
# Check if the source device is preserved
x = jax.device_put(np, jax.devices("cpu")[0])
device = jax.devices("gpu")[0]
y = jax.device_put(x, device)
dl_device = y.__dlpack_device__()
if use_stream:
stream = tuple(y.devices())[0].get_stream_for_external_ready_events()
dlpack = jax.dlpack.to_dlpack(y, copy=copy, stream=stream)
else:
dlpack = jax.dlpack.to_dlpack(y, copy=copy)
z = jax.dlpack.from_dlpack(dlpack)
self.assertEqual(z.devices(), {device})
self.assertAllClose(np.astype(x.dtype), z)
self.assertRaisesRegex(RuntimeError,
"DLPack tensor may be consumed at most once",
lambda: jax.dlpack.from_dlpack(dlpack))
if shape in nonempty_array_shapes:
_check_copy(y, z, bool(copy))
# Check if the destination device can be specified
make_dlpack = lambda: x.__dlpack__(dl_device=dl_device, copy=copy)
if copy == False:
self.assertRaisesRegex(ValueError, "copy=False", make_dlpack)
return
z = jax.dlpack.from_dlpack(make_dlpack())
self.assertEqual(z.devices(), {device})
self.assertAllClose(x, z)
if shape in nonempty_array_shapes:
_check_copy(x, z, True)
@jtu.sample_product(
shape=all_shapes,
dtype=dlpack_dtypes,
gpu=[False, True],
)
def testJaxArrayRoundTrip(self, shape, dtype, gpu):
rng = jtu.rand_default(self.rng())
np = rng(shape, dtype)
if gpu and jax.default_backend() == "cpu":
raise unittest.SkipTest("Skipping GPU test case on CPU")
device = jax.devices("gpu" if gpu else "cpu")[0]
x = jax.device_put(np, device)
y = jax.dlpack.from_dlpack(x)
self.assertEqual(y.devices(), {device})
self.assertAllClose(np.astype(x.dtype), y)
# Test we can create multiple arrays
z = jax.dlpack.from_dlpack(x)
self.assertEqual(z.devices(), {device})
self.assertAllClose(np.astype(x.dtype), z)
@jtu.sample_product(
shape=all_shapes,
dtype=dlpack_dtypes,
)
@unittest.skipIf(not tf, "Test requires TensorFlow")
@jtu.ignore_warning(message="Calling from_dlpack with a DLPack tensor",
category=DeprecationWarning)
def testTensorFlowToJax(self, shape, dtype):
if (not config.enable_x64.value and
dtype in [jnp.int64, jnp.uint64, jnp.float64]):
raise self.skipTest("x64 types are disabled by jax_enable_x64")
if (jtu.test_device_matches(["gpu"]) and
not tf.config.list_physical_devices("GPU")):
raise self.skipTest("TensorFlow not configured with GPU support")
if jtu.test_device_matches(["gpu"]) and dtype == jnp.int32:
raise self.skipTest("TensorFlow does not place int32 tensors on GPU")
rng = jtu.rand_default(self.rng())
np = rng(shape, dtype)
with tf.device("/GPU:0" if jtu.test_device_matches(["gpu"]) else "/CPU:0"):
x = tf.identity(tf.constant(np))
dlpack = tf.experimental.dlpack.to_dlpack(x)
y = jax.dlpack.from_dlpack(dlpack)
self.assertAllClose(np, y)
@jtu.sample_product(
shape=all_shapes,
dtype=dlpack_dtypes,
)
@unittest.skipIf(not tf, "Test requires TensorFlow")
def testJaxToTensorFlow(self, shape, dtype):
if (not config.enable_x64.value and
dtype in [jnp.int64, jnp.uint64, jnp.float64]):
self.skipTest("x64 types are disabled by jax_enable_x64")
if (jtu.test_device_matches(["gpu"]) and
not tf.config.list_physical_devices("GPU")):
raise self.skipTest("TensorFlow not configured with GPU support")
rng = jtu.rand_default(self.rng())
np = rng(shape, dtype)
x = jnp.array(np)
# TODO(b/171320191): this line works around a missing context initialization
# bug in TensorFlow.
_ = tf.add(1, 1)
dlpack = jax.dlpack.to_dlpack(x)
y = tf.experimental.dlpack.from_dlpack(dlpack)
self.assertAllClose(np, y.numpy())
@unittest.skipIf(not tf, "Test requires TensorFlow")
@jtu.ignore_warning(message="Calling from_dlpack with a DLPack tensor",
category=DeprecationWarning)
def testTensorFlowToJaxInt64(self):
# See https://github.com/jax-ml/jax/issues/11895
x = jax.dlpack.from_dlpack(
tf.experimental.dlpack.to_dlpack(tf.ones((2, 3), tf.int64)))
dtype_expected = jnp.int64 if config.enable_x64.value else jnp.int32
self.assertEqual(x.dtype, dtype_expected)
@jtu.sample_product(
shape=all_shapes,
dtype=numpy_dtypes,
copy=[False, True],
)
def testNumpyToJax(self, shape, dtype, copy):
rng = jtu.rand_default(self.rng())
x_np = rng(shape, dtype)
device = jax.devices()[0]
_from_dlpack = lambda: jnp.from_dlpack(x_np, device=device, copy=copy)
if jax.default_backend() == 'gpu' and not copy:
self.assertRaisesRegex(
ValueError,
r"Specified .* which requires a copy",
_from_dlpack
)
else:
self.assertAllClose(x_np, _from_dlpack())
@jtu.sample_product(
shape=all_shapes,
dtype=numpy_dtypes,
)
@jtu.run_on_devices("cpu") # NumPy only accepts cpu DLPacks
def testJaxToNumpy(self, shape, dtype):
rng = jtu.rand_default(self.rng())
x_jax = jnp.array(rng(shape, dtype))
x_np = np.from_dlpack(x_jax)
self.assertAllClose(x_np, x_jax)
@jtu.ignore_warning(message="Calling from_dlpack.*",
category=DeprecationWarning)
def testNondefaultLayout(self):
# Generate numpy array with nonstandard layout
a = np.arange(4).reshape(2, 2)
b = a.T
with self.assertRaisesRegex(
RuntimeError,
r"from_dlpack got array with non-default layout with minor-to-major "
r"dimensions \(0,1\), expected \(1,0\)"):
b_jax = jax.dlpack.from_dlpack(b.__dlpack__())
class CudaArrayInterfaceTest(jtu.JaxTestCase):
@jtu.skip_on_devices("cuda")
def testCudaArrayInterfaceOnNonCudaFails(self):
x = jnp.arange(5)
self.assertFalse(hasattr(x, "__cuda_array_interface__"))
with self.assertRaisesRegex(
AttributeError,
"__cuda_array_interface__ is only defined for NVidia GPU buffers.",
):
_ = x.__cuda_array_interface__
@jtu.run_on_devices("cuda")
def testCudaArrayInterfaceOnShardedArrayFails(self):
devices = jax.local_devices()
if len(devices) <= 1:
raise unittest.SkipTest("Test requires 2 or more devices")
mesh = jax.sharding.Mesh(np.array(devices), ("x",))
sharding = jax.sharding.NamedSharding(mesh, P("x"))
x = jnp.arange(16)
x = jax.device_put(x, sharding)
self.assertFalse(hasattr(x, "__cuda_array_interface__"))
with self.assertRaisesRegex(
AttributeError,
"__cuda_array_interface__ is only supported for unsharded arrays.",
):
_ = x.__cuda_array_interface__
@jtu.sample_product(
shape=all_shapes,
dtype=cuda_array_interface_dtypes,
)
@jtu.run_on_devices("cuda")
def testCudaArrayInterfaceWorks(self, shape, dtype):
rng = jtu.rand_default(self.rng())
x = rng(shape, dtype)
y = jnp.array(x)
z = np.asarray(y)
a = y.__cuda_array_interface__
self.assertEqual(shape, a["shape"])
self.assertEqual(z.__array_interface__["typestr"], a["typestr"])
@jtu.run_on_devices("cuda")
def testCudaArrayInterfaceBfloat16Fails(self):
rng = jtu.rand_default(self.rng())
x = rng((2, 2), jnp.bfloat16)
y = jnp.array(x)
with self.assertRaisesRegex(AttributeError, ".*not supported for BF16.*"):
_ = y.__cuda_array_interface__
@jtu.sample_product(
shape=all_shapes,
dtype=cuda_array_interface_dtypes,
)
@unittest.skipIf(not cupy, "Test requires CuPy")
@jtu.run_on_devices("cuda")
def testJaxToCuPy(self, shape, dtype):
rng = jtu.rand_default(self.rng())
x = rng(shape, dtype)
y = jnp.array(x)
z = cupy.asarray(y)
self.assertEqual(y.__cuda_array_interface__["data"][0],
z.__cuda_array_interface__["data"][0])
self.assertAllClose(x, cupy.asnumpy(z))
@jtu.sample_product(
shape=all_shapes,
dtype=jtu.dtypes.supported(cuda_array_interface_dtypes),
)
@unittest.skipIf(not cupy, "Test requires CuPy")
@jtu.run_on_devices("cuda")
def testCuPyToJax(self, shape, dtype):
rng = jtu.rand_default(self.rng())
x = rng(shape, dtype)
y = cupy.asarray(x)
z = jnp.array(y, copy=False) # this conversion uses dlpack protocol
self.assertEqual(z.dtype, dtype)
self.assertEqual(y.__cuda_array_interface__["data"][0],
z.__cuda_array_interface__["data"][0])
self.assertAllClose(np.asarray(z), cupy.asnumpy(y))
@jtu.sample_product(
shape=all_shapes,
dtype=jtu.dtypes.supported(cuda_array_interface_dtypes),
)
@jtu.run_on_devices("cuda")
def testCaiToJax(self, shape, dtype):
rng = jtu.rand_default(self.rng())
x = rng(shape, dtype)
# using device with highest device_id for testing the correctness
# of detecting the device id from a pointer value
device = jax.devices('cuda')[-1]
with jax.default_device(device):
y = jnp.array(x, dtype=dtype)
self.assertEqual(y.dtype, dtype)
# Using a jax array CAI provider support to construct an object
# that implements the CUDA Array Interface, versions 2 and 3.
cai = y.__cuda_array_interface__
stream = tuple(y.devices())[0].get_stream_for_external_ready_events()
class CAIWithoutStridesV2:
__cuda_array_interface__ = cai.copy()
__cuda_array_interface__["version"] = 2
# CAI version 2 may not define strides and does not define stream
__cuda_array_interface__.pop("strides", None)
__cuda_array_interface__.pop("stream", None)
class CAIWithoutStrides:
__cuda_array_interface__ = cai.copy()
__cuda_array_interface__["version"] = 3
__cuda_array_interface__["strides"] = None
__cuda_array_interface__["stream"] = None # default stream
class CAIWithStrides:
__cuda_array_interface__ = cai.copy()
__cuda_array_interface__["version"] = 3
strides = (dtype.dtype.itemsize,) if shape else ()
for s in reversed(shape[1:]):
strides = (strides[0] * s, *strides)
__cuda_array_interface__['strides'] = strides
__cuda_array_interface__["stream"] = stream
for CAIObject in [CAIWithoutStridesV2, CAIWithoutStrides,
CAIWithStrides]:
z = jnp.array(CAIObject(), copy=False)
self.assertEqual(y.__cuda_array_interface__["data"][0],
z.__cuda_array_interface__["data"][0])
self.assertAllClose(x, z)
if 0 in shape:
# the device id detection from a zero pointer value is not
# possible
pass
else:
self.assertEqual(y.devices(), z.devices())
z = jnp.array(CAIObject(), copy=True)
if 0 not in shape:
self.assertNotEqual(y.__cuda_array_interface__["data"][0],
z.__cuda_array_interface__["data"][0])
self.assertAllClose(x, z)
class Bfloat16Test(jtu.JaxTestCase):
@unittest.skipIf((not tf or tf_version < (2, 5, 0)),
"Test requires TensorFlow 2.5.0 or newer")
def testJaxAndTfHaveTheSameBfloat16Type(self):
self.assertEqual(np.dtype(jnp.bfloat16).num,
np.dtype(tf.dtypes.bfloat16.as_numpy_dtype).num)
if __name__ == "__main__":
absltest.main(testLoader=jtu.JaxTestLoader())