Skip to content

Commit 1272f3c

Browse files
authored
[xpu][test] Port 2 test/dtypes_{floatx, bitpacking} UT files to intel XPU (#3368)
* enable test/dtypes/test_bitpacking.py on intel xpu * enable test/dtypes/test_floatx.py * enable test/dtypes/test_floatx.py * fix format issue * fix format issue * update _DEVICES
1 parent c4273fe commit 1272f3c

File tree

2 files changed

+23
-14
lines changed

2 files changed

+23
-14
lines changed

test/dtypes/test_bitpacking.py

Lines changed: 16 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -8,9 +8,11 @@
88
from torch.utils._triton import has_triton
99

1010
from torchao.dtypes.uintx.bitpacking import pack, pack_cpu, unpack, unpack_cpu
11+
from torchao.utils import get_current_accelerator_device
1112

1213
bit_widths = (1, 2, 3, 4, 5, 6, 7)
1314
dimensions = (0, -1, 1)
15+
_DEVICE = get_current_accelerator_device()
1416

1517

1618
@pytest.fixture(autouse=True)
@@ -30,40 +32,46 @@ def test_CPU(bit_width, dim):
3032
assert unpacked.allclose(test_tensor)
3133

3234

33-
@pytest.mark.skipif(not torch.cuda.is_available(), reason="CUDA not available")
35+
@pytest.mark.skipif(not torch.accelerator.is_available(), reason="GPU not available")
3436
@pytest.mark.parametrize("bit_width", bit_widths)
3537
@pytest.mark.parametrize("dim", dimensions)
3638
def test_GPU(bit_width, dim):
37-
test_tensor = torch.randint(0, 2**bit_width, (32, 32, 32), dtype=torch.uint8).cuda()
39+
test_tensor = torch.randint(0, 2**bit_width, (32, 32, 32), dtype=torch.uint8).to(
40+
_DEVICE
41+
)
3842
packed = pack(test_tensor, bit_width, dim=dim)
3943
unpacked = unpack(packed, bit_width, dim=dim)
4044
assert unpacked.allclose(test_tensor)
4145

4246

43-
@pytest.mark.skipif(not torch.cuda.is_available(), reason="CUDA not available")
47+
@pytest.mark.skipif(not torch.accelerator.is_available(), reason="GPU not available")
4448
@pytest.mark.skipif(not has_triton(), reason="unsupported without triton")
4549
@pytest.mark.parametrize("bit_width", bit_widths)
4650
@pytest.mark.parametrize("dim", dimensions)
4751
def test_compile(bit_width, dim):
4852
torch._dynamo.config.specialize_int = True
4953
torch.compile(pack, fullgraph=True)
5054
torch.compile(unpack, fullgraph=True)
51-
test_tensor = torch.randint(0, 2**bit_width, (32, 32, 32), dtype=torch.uint8).cuda()
55+
test_tensor = torch.randint(0, 2**bit_width, (32, 32, 32), dtype=torch.uint8).to(
56+
_DEVICE
57+
)
5258
packed = pack(test_tensor, bit_width, dim=dim)
5359
unpacked = unpack(packed, bit_width, dim=dim)
5460
assert unpacked.allclose(test_tensor)
5561

5662

5763
# these test cases are for the example pack walk through in the bitpacking.py file
58-
@pytest.mark.skipif(not torch.cuda.is_available(), reason="CUDA not available")
64+
@pytest.mark.skipif(not torch.accelerator.is_available(), reason="GPU not available")
5965
def test_pack_example():
6066
test_tensor = torch.tensor(
6167
[0x30, 0x29, 0x17, 0x5, 0x20, 0x16, 0x9, 0x22], dtype=torch.uint8
62-
).cuda()
68+
).to(_DEVICE)
6369
shard_4, shard_2 = pack(test_tensor, 6)
6470
print(shard_4, shard_2)
65-
assert torch.tensor([0, 105, 151, 37], dtype=torch.uint8).cuda().allclose(shard_4)
66-
assert torch.tensor([39, 146], dtype=torch.uint8).cuda().allclose(shard_2)
71+
assert (
72+
torch.tensor([0, 105, 151, 37], dtype=torch.uint8).to(_DEVICE).allclose(shard_4)
73+
)
74+
assert torch.tensor([39, 146], dtype=torch.uint8).to(_DEVICE).allclose(shard_2)
6775
unpacked = unpack([shard_4, shard_2], 6)
6876
assert unpacked.allclose(test_tensor)
6977

test/dtypes/test_floatx.py

Lines changed: 7 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -33,10 +33,11 @@
3333
quantize_,
3434
)
3535
from torchao.testing.utils import skip_if_rocm
36-
from torchao.utils import is_fbcode
36+
from torchao.utils import get_current_accelerator_device, is_fbcode
3737

38-
_DEVICES = ["cpu"] + (["cuda"] if torch.cuda.is_available() else [])
3938
_Floatx_DTYPES = [(3, 2), (2, 2)]
39+
_DEVICE = get_current_accelerator_device()
40+
_DEVICES = ["cpu"] + ([_DEVICE] if torch.accelerator.is_available() else [])
4041

4142

4243
class TestFloatxTensorCoreAQTTensorImpl(TestCase):
@@ -87,7 +88,7 @@ def test_from_scaled_tc_floatx_compile(self, ebits, mbits, device):
8788
)
8889
torch.testing.assert_close(actual, expected)
8990

90-
@unittest.skipIf(not torch.cuda.is_available(), reason="CUDA not available")
91+
@unittest.skipIf(not torch.accelerator.is_available(), reason="GPU not available")
9192
@parametrize("ebits,mbits", _Floatx_DTYPES)
9293
def test_to_copy_device(self, ebits, mbits):
9394
from torchao.quantization.quant_primitives import (
@@ -101,8 +102,8 @@ def test_to_copy_device(self, ebits, mbits):
101102
_layout = FloatxTensorCoreLayout(ebits, mbits)
102103
floatx_tensor_impl = FloatxTensorCoreAQTTensorImpl.from_plain(
103104
x, scale, None, _layout
104-
).cuda()
105-
assert floatx_tensor_impl.device.type == "cuda"
105+
).to(_DEVICE)
106+
assert floatx_tensor_impl.device.type == _DEVICE.type
106107
floatx_tensor_impl = floatx_tensor_impl.cpu()
107108
assert floatx_tensor_impl.device.type == "cpu"
108109

@@ -114,7 +115,7 @@ def test_to_copy_device(self, ebits, mbits):
114115
@skip_if_rocm("ROCm enablement in progress")
115116
def test_fpx_weight_only(self, ebits, mbits, bias, dtype):
116117
N, OC, IC = 4, 256, 64
117-
device = "cuda"
118+
device = _DEVICE
118119

119120
linear = torch.nn.Linear(IC, OC, bias=bias, device=device, dtype=dtype)
120121
fpx_linear = copy.deepcopy(linear)

0 commit comments

Comments
 (0)