Skip to content

Commit f4bf776

Browse files
authored
Merge branch 'main' into leanfljaenflajenf
2 parents 06bc42c + 375cfdf commit f4bf776

24 files changed

+521
-113
lines changed

.github/workflows/prototype-tests-linux-gpu.yml

+3
Original file line numberDiff line numberDiff line change
@@ -1,5 +1,8 @@
11
name: Prototype tests on Linux
22

3+
# IMPORTANT: This workflow has been manually disabled from the GitHub interface
4+
# in June 2024. The file is kept for reference in case we ever put this back.
5+
36
on:
47
pull_request:
58

docs/source/transforms.rst

+2
Original file line numberDiff line numberDiff line change
@@ -350,6 +350,7 @@ Color
350350
v2.RGB
351351
v2.RandomGrayscale
352352
v2.GaussianBlur
353+
v2.GaussianNoise
353354
v2.RandomInvert
354355
v2.RandomPosterize
355356
v2.RandomSolarize
@@ -368,6 +369,7 @@ Functionals
368369
v2.functional.grayscale_to_rgb
369370
v2.functional.to_grayscale
370371
v2.functional.gaussian_blur
372+
v2.functional.gaussian_noise
371373
v2.functional.invert
372374
v2.functional.posterize
373375
v2.functional.solarize
File renamed without changes.

setup.py

+10-3
Original file line numberDiff line numberDiff line change
@@ -5,6 +5,7 @@
55
import shutil
66
import subprocess
77
import sys
8+
import warnings
89

910
import torch
1011
from pkg_resources import DistributionNotFound, get_distribution, parse_version
@@ -138,7 +139,6 @@ def get_extensions():
138139
+ glob.glob(os.path.join(extensions_dir, "ops", "cpu", "*.cpp"))
139140
+ glob.glob(os.path.join(extensions_dir, "ops", "quantized", "cpu", "*.cpp"))
140141
)
141-
source_mps = glob.glob(os.path.join(extensions_dir, "ops", "mps", "*.mm"))
142142

143143
print("Compiling extensions with following flags:")
144144
force_cuda = os.getenv("FORCE_CUDA", "0") == "1"
@@ -204,8 +204,15 @@ def get_extensions():
204204
define_macros += [("WITH_HIP", None)]
205205
nvcc_flags = []
206206
extra_compile_args["nvcc"] = nvcc_flags
207-
elif torch.backends.mps.is_available() or force_mps:
208-
sources += source_mps
207+
208+
# FIXME: MPS build breaks custom ops registration, so it was disabled.
209+
# See https://github.com/pytorch/vision/issues/8456.
210+
# TODO: Fix MPS build, remove warning below, and put back commented-out elif block.V
211+
if force_mps:
212+
warnings.warn("MPS build is temporarily disabled!!!!")
213+
# elif torch.backends.mps.is_available() or force_mps:
214+
# source_mps = glob.glob(os.path.join(extensions_dir, "ops", "mps", "*.mm"))
215+
# sources += source_mps
209216

210217
if sys.platform == "win32":
211218
define_macros += [("torchvision_EXPORTS", None)]

test/conftest.py

+2-1
Original file line numberDiff line numberDiff line change
@@ -49,7 +49,8 @@ def pytest_collection_modifyitems(items):
4949
# There are special cases though, see below
5050
item.add_marker(pytest.mark.skip(reason=CUDA_NOT_AVAILABLE_MSG))
5151

52-
if needs_mps and not torch.backends.mps.is_available():
52+
# TODO: uncoment when MPS works again - see FIXME in setup.py
53+
if needs_mps: # and not torch.backends.mps.is_available():
5354
item.add_marker(pytest.mark.skip(reason=MPS_NOT_AVAILABLE_MSG))
5455

5556
if IN_FBCODE:

test/test_datasets.py

+92-22
Original file line numberDiff line numberDiff line change
@@ -782,32 +782,46 @@ def inject_fake_data(self, tmpdir, config):
782782

783783
annotation_folder = tmpdir / self._ANNOTATIONS_FOLDER
784784
os.makedirs(annotation_folder)
785+
786+
segmentation_kind = config.pop("segmentation_kind", "list")
785787
info = self._create_annotation_file(
786-
annotation_folder, self._ANNOTATIONS_FILE, file_names, num_annotations_per_image
788+
annotation_folder,
789+
self._ANNOTATIONS_FILE,
790+
file_names,
791+
num_annotations_per_image,
792+
segmentation_kind=segmentation_kind,
787793
)
788794

789795
info["num_examples"] = num_images
790796
return info
791797

792-
def _create_annotation_file(self, root, name, file_names, num_annotations_per_image):
798+
def _create_annotation_file(self, root, name, file_names, num_annotations_per_image, segmentation_kind="list"):
793799
image_ids = [int(file_name.stem) for file_name in file_names]
794800
images = [dict(file_name=str(file_name), id=id) for file_name, id in zip(file_names, image_ids)]
795801

796-
annotations, info = self._create_annotations(image_ids, num_annotations_per_image)
802+
annotations, info = self._create_annotations(image_ids, num_annotations_per_image, segmentation_kind)
797803
self._create_json(root, name, dict(images=images, annotations=annotations))
798804

799805
return info
800806

801-
def _create_annotations(self, image_ids, num_annotations_per_image):
807+
def _create_annotations(self, image_ids, num_annotations_per_image, segmentation_kind="list"):
802808
annotations = []
803809
annotion_id = 0
810+
804811
for image_id in itertools.islice(itertools.cycle(image_ids), len(image_ids) * num_annotations_per_image):
812+
segmentation = {
813+
"list": [torch.rand(8).tolist()],
814+
"rle": {"size": [10, 10], "counts": [1]},
815+
"rle_encoded": {"size": [2400, 2400], "counts": "PQRQ2[1\\Y2f0gNVNRhMg2"},
816+
"bad": 123,
817+
}[segmentation_kind]
818+
805819
annotations.append(
806820
dict(
807821
image_id=image_id,
808822
id=annotion_id,
809823
bbox=torch.rand(4).tolist(),
810-
segmentation=[torch.rand(8).tolist()],
824+
segmentation=segmentation,
811825
category_id=int(torch.randint(91, ())),
812826
area=float(torch.rand(1)),
813827
iscrowd=int(torch.randint(2, size=(1,))),
@@ -832,11 +846,27 @@ def test_slice_error(self):
832846
with pytest.raises(ValueError, match="Index must be of type integer"):
833847
dataset[:2]
834848

849+
def test_segmentation_kind(self):
850+
if isinstance(self, CocoCaptionsTestCase):
851+
return
852+
853+
for segmentation_kind in ("list", "rle", "rle_encoded"):
854+
config = {"segmentation_kind": segmentation_kind}
855+
with self.create_dataset(config) as (dataset, _):
856+
dataset = datasets.wrap_dataset_for_transforms_v2(dataset, target_keys="all")
857+
list(dataset)
858+
859+
config = {"segmentation_kind": "bad"}
860+
with self.create_dataset(config) as (dataset, _):
861+
dataset = datasets.wrap_dataset_for_transforms_v2(dataset, target_keys="all")
862+
with pytest.raises(ValueError, match="COCO segmentation expected to be a dict or a list"):
863+
list(dataset)
864+
835865

836866
class CocoCaptionsTestCase(CocoDetectionTestCase):
837867
DATASET_CLASS = datasets.CocoCaptions
838868

839-
def _create_annotations(self, image_ids, num_annotations_per_image):
869+
def _create_annotations(self, image_ids, num_annotations_per_image, segmentation_kind="list"):
840870
captions = [str(idx) for idx in range(num_annotations_per_image)]
841871
annotations = combinations_grid(image_id=image_ids, caption=captions)
842872
for id, annotation in enumerate(annotations):
@@ -2442,28 +2472,68 @@ def inject_fake_data(self, tmpdir, config):
24422472
base_folder = os.path.join(tmpdir, "fer2013")
24432473
os.makedirs(base_folder)
24442474

2475+
use_icml = config.pop("use_icml", False)
2476+
use_fer = config.pop("use_fer", False)
2477+
24452478
num_samples = 5
2446-
with open(os.path.join(base_folder, f"{config['split']}.csv"), "w", newline="") as file:
2447-
writer = csv.DictWriter(
2448-
file,
2449-
fieldnames=("emotion", "pixels") if config["split"] == "train" else ("pixels",),
2450-
quoting=csv.QUOTE_NONNUMERIC,
2451-
quotechar='"',
2452-
)
2453-
writer.writeheader()
2454-
for _ in range(num_samples):
2455-
row = dict(
2456-
pixels=" ".join(
2457-
str(pixel) for pixel in datasets_utils.create_image_or_video_tensor((48, 48)).view(-1).tolist()
2458-
)
2479+
2480+
if use_icml or use_fer:
2481+
pixels_key, usage_key = (" pixels", " Usage") if use_icml else ("pixels", "Usage")
2482+
fieldnames = ("emotion", usage_key, pixels_key) if use_icml else ("emotion", pixels_key, usage_key)
2483+
filename = "icml_face_data.csv" if use_icml else "fer2013.csv"
2484+
with open(os.path.join(base_folder, filename), "w", newline="") as file:
2485+
writer = csv.DictWriter(
2486+
file,
2487+
fieldnames=fieldnames,
2488+
quoting=csv.QUOTE_NONNUMERIC,
2489+
quotechar='"',
2490+
)
2491+
writer.writeheader()
2492+
for i in range(num_samples):
2493+
row = {
2494+
"emotion": str(int(torch.randint(0, 7, ()))),
2495+
usage_key: "Training" if i % 2 else "PublicTest",
2496+
pixels_key: " ".join(
2497+
str(pixel)
2498+
for pixel in datasets_utils.create_image_or_video_tensor((48, 48)).view(-1).tolist()
2499+
),
2500+
}
2501+
2502+
writer.writerow(row)
2503+
else:
2504+
with open(os.path.join(base_folder, f"{config['split']}.csv"), "w", newline="") as file:
2505+
writer = csv.DictWriter(
2506+
file,
2507+
fieldnames=("emotion", "pixels") if config["split"] == "train" else ("pixels",),
2508+
quoting=csv.QUOTE_NONNUMERIC,
2509+
quotechar='"',
24592510
)
2460-
if config["split"] == "train":
2461-
row["emotion"] = str(int(torch.randint(0, 7, ())))
2511+
writer.writeheader()
2512+
for _ in range(num_samples):
2513+
row = dict(
2514+
pixels=" ".join(
2515+
str(pixel)
2516+
for pixel in datasets_utils.create_image_or_video_tensor((48, 48)).view(-1).tolist()
2517+
)
2518+
)
2519+
if config["split"] == "train":
2520+
row["emotion"] = str(int(torch.randint(0, 7, ())))
24622521

2463-
writer.writerow(row)
2522+
writer.writerow(row)
24642523

24652524
return num_samples
24662525

2526+
def test_icml_file(self):
2527+
config = {"split": "test"}
2528+
with self.create_dataset(config=config) as (dataset, _):
2529+
assert all(s[1] is None for s in dataset)
2530+
2531+
for split in ("train", "test"):
2532+
for d in ({"use_icml": True}, {"use_fer": True}):
2533+
config = {"split": split, **d}
2534+
with self.create_dataset(config=config) as (dataset, _):
2535+
assert all(s[1] is not None for s in dataset)
2536+
24672537

24682538
class GTSRBTestCase(datasets_utils.ImageDatasetTestCase):
24692539
DATASET_CLASS = datasets.GTSRB

0 commit comments

Comments
 (0)