Skip to content

Commit 1e4cba7

Browse files
committed
Add support for max_version, dl_device, copy kwargs in __dlpack__
1 parent 477a44f commit 1e4cba7

File tree

3 files changed

+148
-31
lines changed

3 files changed

+148
-31
lines changed

jax/_src/array.py

Lines changed: 18 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -404,11 +404,25 @@ def __array__(self, dtype=None, context=None, copy=None):
404404
kwds = {} if copy is None else {'copy': copy}
405405
return np.asarray(self._value, dtype=dtype, **kwds)
406406

407-
def __dlpack__(self, *, stream: int | Any | None = None):
408-
if len(self._arrays) != 1:
409-
raise BufferError("__dlpack__ only supported for unsharded arrays.")
407+
def __dlpack__(self, *, stream: int | Any | None = None,
408+
max_version: tuple[int, int] | None = None,
409+
dl_device: tuple[enum.Enum, int] | None = None,
410+
copy: bool | None = None):
410411
from jax._src.dlpack import to_dlpack # pylint: disable=g-import-not-at-top
411-
return to_dlpack(self, stream=stream)
412+
413+
device_set = self.sharding.device_set
414+
if len(device_set) > 1:
415+
raise BufferError(
416+
"to_dlpack can only pack a dlpack tensor from an array on a singular "
417+
f"device, but an array with a Sharding over {len(device_set)} devices "
418+
"was provided."
419+
)
420+
device, = device_set
421+
return to_dlpack(self, stream=stream,
422+
max_version=max_version,
423+
src_device=device,
424+
dl_device=dl_device, # type: ignore
425+
copy=copy)
412426

413427
def __dlpack_device__(self) -> tuple[enum.Enum, int]:
414428
if len(self._arrays) != 1:

jax/_src/dlpack.py

Lines changed: 93 additions & 17 deletions
Original file line numberDiff line numberDiff line change
@@ -16,17 +16,20 @@
1616

1717
import enum
1818
from typing import Any
19-
import warnings
2019

2120
from jax._src.api import device_put
2221
from jax import numpy as jnp
2322
from jax._src import array
2423
from jax._src import xla_bridge
24+
from jax._src.lax.lax import _array_copy
2525
from jax._src.lib import xla_client
2626
from jax._src.lib import xla_extension_version
2727
from jax._src.typing import Array
2828
from jax._src.sharding import Sharding
2929

30+
DLPACK_VERSION = (0, 8)
31+
MIN_DLPACK_VERSION = (0, 5)
32+
3033
# A set of dtypes that dlpack supports.
3134
# Note: Make sure to use a "type", not a dtype instance, when looking up this set
3235
# because their hashes are different.
@@ -49,39 +52,112 @@ class DLDeviceType(enum.IntEnum):
4952
kDLCUDA = 2
5053
kDLROCM = 10
5154

55+
def _to_dlpack(x: Array, stream: int | Any | None,
56+
src_device: xla_client.Device | None = None,
57+
device: xla_client.Device | None = None,
58+
copy: bool | None = None):
59+
60+
if src_device is None:
61+
src_device, = x.devices()
62+
if device and (src_device is None or device != src_device):
63+
if copy is not None and not copy:
64+
raise ValueError(
65+
f"Specified {device=} which requires a copy since the source device "
66+
f"is {repr(src_device)}, however copy=False. Set copy=True or "
67+
"copy=None to perform the requested operation."
68+
)
69+
else:
70+
arr = device_put(x, device)
71+
else:
72+
arr = _array_copy(x) if copy else x
73+
return xla_client._xla.buffer_to_dlpack_managed_tensor(
74+
arr.addressable_data(0), stream=stream
75+
)
5276

53-
def to_dlpack(x: Array, take_ownership: bool = False,
54-
stream: int | Any | None = None):
77+
def to_dlpack(x: Array, stream: int | Any | None = None,
78+
src_device: xla_client.Device | None = None,
79+
dl_device: tuple[DLDeviceType, int] | None = None,
80+
max_version: tuple[int, int] | None = None,
81+
copy : bool | None = None):
5582
"""Returns a DLPack tensor that encapsulates a :class:`~jax.Array` ``x``.
5683
5784
Args:
5885
x: a :class:`~jax.Array`, on either CPU or GPU.
59-
take_ownership: Deprecated. It is a no-op to set take_ownership. Will be
60-
deleted in 01/2024.
6186
stream: optional platform-dependent stream to wait on until the buffer is
6287
ready. This corresponds to the `stream` argument to ``__dlpack__``
6388
documented in https://dmlc.github.io/dlpack/latest/python_spec.html.
89+
src_device: either a CPU or GPU :class:`~jax.Device`.
90+
dl_device: a tuple of ``(dl_device_type, local_hardware_id)`` in DLPack
91+
format e.g. as produced by ``__dlpack_device__``.
92+
max_version: the maximum DLPack version that the consumer (i.e. caller of
93+
``__dlpack__``) supports in the form of a 2-tuple of ``(major, minor)``.
94+
This function is not guaranteed to return a capsule of version
95+
``max_version``.
96+
copy: a boolean indicating whether or not to copy the input. If
97+
``copy=True`` then the function must always copy. When
98+
``copy=False`` then the function must never copy, and must raise an error
99+
when a copy is deemed necessary. If ``copy=None`` then the function must
100+
avoid a copy if possible but may copy if needed.
64101
65102
Returns:
66-
A dlpack PyCapsule object.
103+
A DLPack PyCapsule object.
67104
68105
Note:
69-
While JAX arrays are always immutable, dlpack buffers cannot be marked as
70-
immutable, and it is possible for processes external to JAX to mutate them
71-
in-place. If a dlpack buffer derived from a JAX array is mutated, it may
72-
lead to undefined behavior when using the associated JAX array.
106+
While JAX arrays are always immutable, ``DLPackManagedTensor`` buffers
107+
cannot be marked as immutable, and it is possible for processes external
108+
to JAX to mutate them in-place. If a DLPack buffer derived from a JAX array
109+
is mutated, it may lead to undefined behavior when using the associated JAX
110+
array. When JAX eventually supports ``DLManagedTensorVersioned``
111+
(DLPack 1.0), it will be possible to specify that a buffer is read-only.
73112
"""
74113
if not isinstance(x, array.ArrayImpl):
75114
raise TypeError("Argument to to_dlpack must be a jax.Array, "
76115
f"got {type(x)}")
77-
assert len(x.devices()) == 1
78-
if take_ownership:
79-
warnings.warn(
80-
"take_ownership in to_dlpack is deprecated and it is a no-op."
116+
117+
device = None
118+
dl_device_type, local_hardware_id = dl_device if dl_device else (None, None)
119+
if dl_device_type:
120+
try:
121+
dl_device_platform = {
122+
DLDeviceType.kDLCPU: "cpu",
123+
DLDeviceType.kDLCUDA: "cuda",
124+
DLDeviceType.kDLROCM: "rocm",
125+
}[dl_device_type]
126+
backend = xla_bridge.get_backend(dl_device_platform)
127+
device = backend.device_from_local_hardware_id(local_hardware_id)
128+
except TypeError:
129+
# https://data-apis.org/array-api/latest/API_specification/generated/array_api.array.__dlpack__.html
130+
# recommends using BufferError.
131+
raise BufferError(
132+
"The device specification passed to to_dlpack contains an unsupported "
133+
f"device type (DLDeviceType: {dl_device_type})")
134+
135+
# As new versions are adopted over time, we can maintain some legacy paths
136+
# for compatability mediated through the max_version parameter.
137+
# TODO(micky774): Deprecate default usage of DLPackManagedTensor when XLA
138+
# supports DLManagedTensorVersioned (DLPack version 1.0) and repurpose the
139+
# current _to_dlpack as a legacy path for (0,5) <= max_version < (1,0).
140+
if max_version is None or max_version >= DLPACK_VERSION:
141+
# Latest
142+
return _to_dlpack(
143+
x, stream=stream,
144+
src_device=src_device,
145+
device=device,
146+
copy=copy
147+
)
148+
elif max_version >= MIN_DLPACK_VERSION:
149+
# Oldest supported
150+
return _to_dlpack(
151+
x, stream=stream,
152+
src_device=src_device,
153+
device=device,
154+
copy=copy
155+
)
156+
else:
157+
raise BufferError(
158+
f"JAX does not support any version below {MIN_DLPACK_VERSION} but "
159+
f"version ({max_version}) was requested."
81160
)
82-
return xla_client._xla.buffer_to_dlpack_managed_tensor(
83-
x.addressable_data(0), stream=stream
84-
) # type: ignore
85161

86162
def _place_array(_arr, device, dlpack_device, copy):
87163
if device and dlpack_device != device:

tests/array_interoperability_test.py

Lines changed: 37 additions & 10 deletions
Original file line numberDiff line numberDiff line change
@@ -74,22 +74,49 @@ def setUp(self):
7474
shape=all_shapes,
7575
dtype=dlpack_dtypes,
7676
gpu=[False, True],
77+
copy=[False, True, None]
7778
)
78-
def testJaxRoundTrip(self, shape, dtype, gpu):
79+
def testJaxRoundTrip(self, shape, dtype, gpu, copy):
80+
if xb.using_pjrt_c_api():
81+
self.skipTest("DLPack support is incomplete in the PJRT C API") # TODO(skyewm)
7982
rng = jtu.rand_default(self.rng())
8083
np = rng(shape, dtype)
81-
if gpu and jtu.test_device_matches(["cpu"]):
82-
raise unittest.SkipTest("Skipping GPU test case on CPU")
84+
if gpu and not jtu.test_device_matches(["gpu"]):
85+
raise unittest.SkipTest("Skipping GPU test case on CPU/TPU")
86+
87+
def _check_copy(x, y, expect_copy):
88+
is_copy = x.unsafe_buffer_pointer() != y.unsafe_buffer_pointer()
89+
assert is_copy == expect_copy
90+
91+
# Check if the source device is preserved
92+
x = jax.device_put(np, jax.devices("cpu")[0])
8393
device = jax.devices("gpu" if gpu else "cpu")[0]
84-
x = jax.device_put(np, device)
85-
dlpack = jax.dlpack.to_dlpack(x)
86-
y = jax.dlpack.from_dlpack(dlpack)
87-
self.assertEqual(y.devices(), {device})
88-
self.assertAllClose(np.astype(x.dtype), y)
94+
y = jax.device_put(x, device)
95+
dlpack = jax.dlpack.to_dlpack(y, copy=copy)
96+
z = jax.dlpack.from_dlpack(dlpack)
8997

98+
self.assertEqual(z.devices(), {device})
99+
self.assertAllClose(np.astype(x.dtype), z)
90100
self.assertRaisesRegex(RuntimeError,
91-
"DLPack tensor may be consumed at most once",
92-
lambda: jax.dlpack.from_dlpack(dlpack))
101+
"DLPack tensor may be consumed at most once",
102+
lambda: jax.dlpack.from_dlpack(dlpack))
103+
104+
if shape in nonempty_array_shapes:
105+
_check_copy(y, z, bool(copy))
106+
107+
# Check if the destination device can be specified
108+
dl_device = y.__dlpack_device__()
109+
make_dlpack = lambda: x.__dlpack__(dl_device=dl_device, copy=copy)
110+
if gpu and copy == False:
111+
self.assertRaisesRegex(ValueError, "copy=False", make_dlpack)
112+
return
113+
114+
z = jax.dlpack.from_dlpack(make_dlpack())
115+
self.assertEqual(z.devices(), {device})
116+
self.assertAllClose(np.astype(x.dtype), z)
117+
118+
if shape in nonempty_array_shapes:
119+
_check_copy(x, z, bool(copy) or gpu)
93120

94121
@jtu.sample_product(
95122
shape=all_shapes,

0 commit comments

Comments
 (0)