Skip to content

Commit d561f19

Browse files
shapovalovfacebook-github-bot
authored andcommitted
Cleaning up camera difficulty
Summary: We don’t see much value in reporting metrics by camera difficulty while supporting that in new datasets is quite painful, hence deprecating training cameras in the data API and ignoring in evaluation. Reviewed By: bottler Differential Revision: D42678879 fbshipit-source-id: aad511f6cb2ca82745f31c19594e1d80594b61d7
1 parent 1de2d0c commit d561f19

File tree

7 files changed

+22
-106
lines changed

7 files changed

+22
-106
lines changed

projects/implicitron_trainer/experiment.py

-7
Original file line numberDiff line numberDiff line change
@@ -207,12 +207,6 @@ def run(self) -> None:
207207
val_loader,
208208
) = accelerator.prepare(model, optimizer, train_loader, val_loader)
209209

210-
# pyre-fixme[16]: Optional type has no attribute `is_multisequence`.
211-
if not self.training_loop.evaluator.is_multisequence:
212-
all_train_cameras = self.data_source.all_train_cameras
213-
else:
214-
all_train_cameras = None
215-
216210
# Enter the main training loop.
217211
self.training_loop.run(
218212
train_loader=train_loader,
@@ -223,7 +217,6 @@ def run(self) -> None:
223217
model=model,
224218
optimizer=optimizer,
225219
scheduler=scheduler,
226-
all_train_cameras=all_train_cameras,
227220
accelerator=accelerator,
228221
device=device,
229222
exp_dir=self.exp_dir,

projects/implicitron_trainer/impl/training_loop.py

-4
Original file line numberDiff line numberDiff line change
@@ -122,7 +122,6 @@ def run(
122122
optimizer: torch.optim.Optimizer,
123123
scheduler: Any,
124124
accelerator: Optional[Accelerator],
125-
all_train_cameras: Optional[CamerasBase],
126125
device: torch.device,
127126
exp_dir: str,
128127
stats: Stats,
@@ -142,7 +141,6 @@ def run(
142141
if test_loader is not None:
143142
# pyre-fixme[16]: `Optional` has no attribute `run`.
144143
self.evaluator.run(
145-
all_train_cameras=all_train_cameras,
146144
dataloader=test_loader,
147145
device=device,
148146
dump_to_json=True,
@@ -200,7 +198,6 @@ def run(
200198
and epoch % self.test_interval == 0
201199
):
202200
self.evaluator.run(
203-
all_train_cameras=all_train_cameras,
204201
device=device,
205202
dataloader=test_loader,
206203
model=model,
@@ -217,7 +214,6 @@ def run(
217214
if self.test_when_finished:
218215
if test_loader is not None:
219216
self.evaluator.run(
220-
all_train_cameras=all_train_cameras,
221217
device=device,
222218
dump_to_json=True,
223219
epoch=stats.epoch,

pytorch3d/implicitron/dataset/data_source.py

+4
Original file line numberDiff line numberDiff line change
@@ -34,6 +34,7 @@ def get_datasets_and_dataloaders(self) -> Tuple[DatasetMap, DataLoaderMap]:
3434
@property
3535
def all_train_cameras(self) -> Optional[CamerasBase]:
3636
"""
37+
DEPRECATED! The property will be removed in future versions.
3738
If the data is all for a single scene, a list
3839
of the known training cameras for that scene, which is
3940
used for evaluating the viewpoint difficulty of the
@@ -70,6 +71,9 @@ def get_datasets_and_dataloaders(self) -> Tuple[DatasetMap, DataLoaderMap]:
7071

7172
@property
7273
def all_train_cameras(self) -> Optional[CamerasBase]:
74+
"""
75+
DEPRECATED! The property will be removed in future versions.
76+
"""
7377
if self._all_train_cameras_cache is None: # pyre-ignore[16]
7478
all_train_cameras = self.dataset_map_provider.get_all_train_cameras()
7579
self._all_train_cameras_cache = (all_train_cameras,)

pytorch3d/implicitron/dataset/dataset_map_provider.py

+1
Original file line numberDiff line numberDiff line change
@@ -95,6 +95,7 @@ def get_dataset_map(self) -> DatasetMap:
9595

9696
def get_all_train_cameras(self) -> Optional[CamerasBase]:
9797
"""
98+
DEPRECATED! The function will be removed in future versions.
9899
If the data is all for a single scene, returns a list
99100
of the known training cameras for that scene, which is
100101
used for evaluating the difficulty of the unknown

pytorch3d/implicitron/eval_demo.py

+1-10
Original file line numberDiff line numberDiff line change
@@ -153,21 +153,12 @@ def evaluate_dbir_for_category(
153153
preds["implicitron_render"],
154154
bg_color=bg_color,
155155
lpips_model=lpips_model,
156-
source_cameras=data_source.all_train_cameras,
157156
)
158157
)
159158

160-
if task == Task.SINGLE_SEQUENCE:
161-
camera_difficulty_bin_breaks = 0.97, 0.98
162-
multisequence_evaluation = False
163-
else:
164-
camera_difficulty_bin_breaks = 2.0 / 3, 5.0 / 6
165-
multisequence_evaluation = True
166-
167159
category_result_flat, category_result = summarize_nvs_eval_results(
168160
per_batch_eval_results,
169-
camera_difficulty_bin_breaks=camera_difficulty_bin_breaks,
170-
is_multisequence=multisequence_evaluation,
161+
is_multisequence=task != Task.SINGLE_SEQUENCE,
171162
)
172163

173164
return category_result["results"]

pytorch3d/implicitron/evaluation/evaluate_new_view_synthesis.py

+15-74
Original file line numberDiff line numberDiff line change
@@ -18,13 +18,11 @@
1818
from pytorch3d.implicitron.dataset.utils import is_known_frame, is_train_frame
1919
from pytorch3d.implicitron.models.base_model import ImplicitronRender
2020
from pytorch3d.implicitron.tools import vis_utils
21-
from pytorch3d.implicitron.tools.camera_utils import volumetric_camera_overlaps
2221
from pytorch3d.implicitron.tools.image_utils import mask_background
2322
from pytorch3d.implicitron.tools.metric_utils import calc_psnr, eval_depth, iou, rgb_l1
2423
from pytorch3d.implicitron.tools.point_cloud_utils import get_rgbd_point_cloud
2524
from pytorch3d.implicitron.tools.vis_utils import make_depth_image
26-
from pytorch3d.renderer.camera_utils import join_cameras_as_batch
27-
from pytorch3d.renderer.cameras import CamerasBase, PerspectiveCameras
25+
from pytorch3d.renderer.cameras import PerspectiveCameras
2826
from pytorch3d.vis.plotly_vis import plot_scene
2927
from tabulate import tabulate
3028

@@ -149,7 +147,6 @@ def eval_batch(
149147
visualize: bool = False,
150148
visualize_visdom_env: str = "eval_debug",
151149
break_after_visualising: bool = True,
152-
source_cameras: Optional[CamerasBase] = None,
153150
) -> Dict[str, Any]:
154151
"""
155152
Produce performance metrics for a single batch of new-view synthesis
@@ -171,8 +168,6 @@ def eval_batch(
171168
ground truth.
172169
lpips_model: A pre-trained model for evaluating the LPIPS metric.
173170
visualize: If True, visualizes the results to Visdom.
174-
source_cameras: A list of all training cameras for evaluating the
175-
difficulty of the target views.
176171
177172
Returns:
178173
results: A dictionary holding evaluation metrics.
@@ -365,16 +360,7 @@ def eval_batch(
365360
# convert all metrics to floats
366361
results = {k: float(v) for k, v in results.items()}
367362

368-
if source_cameras is None:
369-
# pyre-fixme[16]: Optional has no attribute __getitem__
370-
source_cameras = frame_data.camera[torch.where(is_known)[0]]
371-
372363
results["meta"] = {
373-
# calculate the camera difficulties and add to results
374-
"camera_difficulty": calculate_camera_difficulties(
375-
frame_data.camera[0],
376-
source_cameras,
377-
)[0].item(),
378364
# store the size of the batch (corresponds to n_src_views+1)
379365
"batch_size": int(is_known.numel()),
380366
# store the type of the target frame
@@ -406,33 +392,6 @@ def average_per_batch_results(
406392
}
407393

408394

409-
def calculate_camera_difficulties(
410-
cameras_target: CamerasBase,
411-
cameras_source: CamerasBase,
412-
) -> torch.Tensor:
413-
"""
414-
Calculate the difficulties of the target cameras, given a set of known
415-
cameras `cameras_source`.
416-
417-
Returns:
418-
a tensor of shape (len(cameras_target),)
419-
"""
420-
ious = [
421-
volumetric_camera_overlaps(
422-
join_cameras_as_batch(
423-
# pyre-fixme[6]: Expected `CamerasBase` for 1st param but got
424-
# `Optional[pytorch3d.renderer.utils.TensorProperties]`.
425-
[cameras_target[cami], cameras_source.to(cameras_target.device)]
426-
)
427-
)[0, :]
428-
for cami in range(cameras_target.R.shape[0])
429-
]
430-
camera_difficulties = torch.stack(
431-
[_reduce_camera_iou_overlap(iou[1:]) for iou in ious]
432-
)
433-
return camera_difficulties
434-
435-
436395
def _reduce_camera_iou_overlap(ious: torch.Tensor, topk: int = 2) -> torch.Tensor:
437396
"""
438397
Calculate the final camera difficulty by computing the average of the
@@ -458,8 +417,7 @@ def _get_camera_difficulty_bin_edges(camera_difficulty_bin_breaks: Tuple[float,
458417
def summarize_nvs_eval_results(
459418
per_batch_eval_results: List[Dict[str, Any]],
460419
is_multisequence: bool,
461-
camera_difficulty_bin_breaks: Tuple[float, float],
462-
):
420+
) -> Tuple[Dict[str, Any], Dict[str, Any]]:
463421
"""
464422
Compile the per-batch evaluation results `per_batch_eval_results` into
465423
a set of aggregate metrics. The produced metrics depend on is_multisequence.
@@ -482,19 +440,12 @@ def summarize_nvs_eval_results(
482440
batch_sizes = torch.tensor(
483441
[r["meta"]["batch_size"] for r in per_batch_eval_results]
484442
).long()
485-
camera_difficulty = torch.tensor(
486-
[r["meta"]["camera_difficulty"] for r in per_batch_eval_results]
487-
).float()
443+
488444
is_train = is_train_frame([r["meta"]["frame_type"] for r in per_batch_eval_results])
489445

490446
# init the result database dict
491447
results = []
492448

493-
diff_bin_edges, diff_bin_names = _get_camera_difficulty_bin_edges(
494-
camera_difficulty_bin_breaks
495-
)
496-
n_diff_edges = diff_bin_edges.numel()
497-
498449
# add per set averages
499450
for SET in eval_sets:
500451
if SET is None:
@@ -504,26 +455,17 @@ def summarize_nvs_eval_results(
504455
ok_set = is_train == int(SET == "train")
505456
set_name = SET
506457

507-
# eval each difficulty bin, including a full average result (diff_bin=None)
508-
for diff_bin in [None, *list(range(n_diff_edges - 1))]:
509-
if diff_bin is None:
510-
# average over all results
511-
in_bin = ok_set
512-
diff_bin_name = "all"
513-
else:
514-
b1, b2 = diff_bin_edges[diff_bin : (diff_bin + 2)]
515-
in_bin = ok_set & (camera_difficulty > b1) & (camera_difficulty <= b2)
516-
diff_bin_name = diff_bin_names[diff_bin]
517-
bin_results = average_per_batch_results(
518-
per_batch_eval_results, idx=torch.where(in_bin)[0]
519-
)
520-
results.append(
521-
{
522-
"subset": set_name,
523-
"subsubset": f"diff={diff_bin_name}",
524-
"metrics": bin_results,
525-
}
526-
)
458+
# average over all results
459+
bin_results = average_per_batch_results(
460+
per_batch_eval_results, idx=torch.where(ok_set)[0]
461+
)
462+
results.append(
463+
{
464+
"subset": set_name,
465+
"subsubset": "diff=all",
466+
"metrics": bin_results,
467+
}
468+
)
527469

528470
if is_multisequence:
529471
# split based on n_src_views
@@ -552,7 +494,7 @@ def _get_flat_nvs_metric_key(result, metric_name) -> str:
552494
return metric_key
553495

554496

555-
def flatten_nvs_results(results):
497+
def flatten_nvs_results(results) -> Dict[str, Any]:
556498
"""
557499
Takes input `results` list of dicts of the form::
558500
@@ -571,7 +513,6 @@ def flatten_nvs_results(results):
571513
'subset=train/test/...|subsubset=src=1/src=2/...': nvs_eval_metrics,
572514
...
573515
}
574-
575516
"""
576517
results_flat = {}
577518
for result in results:

pytorch3d/implicitron/evaluation/evaluator.py

+1-11
Original file line numberDiff line numberDiff line change
@@ -23,7 +23,6 @@
2323
ReplaceableBase,
2424
run_auto_creation,
2525
)
26-
from pytorch3d.renderer.cameras import CamerasBase
2726
from torch.utils.data import DataLoader
2827

2928
logger = logging.getLogger(__name__)
@@ -50,12 +49,9 @@ def run(
5049
class ImplicitronEvaluator(EvaluatorBase):
5150
"""
5251
Evaluate the results of Implicitron training.
53-
54-
Members:
55-
camera_difficulty_bin_breaks: low/medium vals to divide camera difficulties into
56-
[0-eps, low, medium, 1+eps].
5752
"""
5853

54+
# UNUSED; preserved for compatibility purposes
5955
camera_difficulty_bin_breaks: Tuple[float, ...] = 0.97, 0.98
6056

6157
def __post_init__(self):
@@ -65,7 +61,6 @@ def run(
6561
self,
6662
model: ImplicitronModelBase,
6763
dataloader: DataLoader,
68-
all_train_cameras: Optional[CamerasBase],
6964
device: torch.device,
7065
dump_to_json: bool = False,
7166
exp_dir: Optional[str] = None,
@@ -79,7 +74,6 @@ def run(
7974
Args:
8075
model: A (trained) model to evaluate.
8176
dataloader: A test dataloader.
82-
all_train_cameras: Camera instances we used for training.
8377
device: A torch device.
8478
dump_to_json: If True, will dump the results to a json file.
8579
exp_dir: Root expeirment directory.
@@ -123,16 +117,12 @@ def run(
123117
implicitron_render,
124118
bg_color="black",
125119
lpips_model=lpips_model,
126-
source_cameras=( # None will make it use batch’s known cameras
127-
None if self.is_multisequence else all_train_cameras
128-
),
129120
)
130121
)
131122

132123
_, category_result = evaluate.summarize_nvs_eval_results(
133124
per_batch_eval_results,
134125
self.is_multisequence,
135-
self.camera_difficulty_bin_breaks,
136126
)
137127

138128
results = category_result["results"]

0 commit comments

Comments
 (0)